loader.py 5.8 KB

123456789101112131415161718192021222324252627282930313233343536373839404142434445464748495051525354555657585960616263646566676869707172737475767778798081828384858687888990919293949596979899100101102103104105106107108109110111112113114115116117118119120121122123124125126127128129130131132133134135136137138139140141142143144145146147148
  1. """Utilities for selecting and loading models."""
  2. import contextlib
  3. import gc
  4. from contextlib import nullcontext
  5. from typing import Type
  6. from loguru import logger
  7. import torch
  8. import torch.nn as nn
  9. from aphrodite.common.config import DeviceConfig, ModelConfig
  10. from aphrodite.modeling.models import ModelRegistry
  11. from aphrodite.modeling.models.llava import LlavaForConditionalGeneration
  12. from aphrodite.modeling.hf_downloader import (
  13. get_quant_config,
  14. initialize_dummy_weights,
  15. )
  16. from aphrodite.modeling.layers.quantization.bitsandbytes import (
  17. BNBLinearMethod,
  18. replace_quant_params,
  19. )
  20. from aphrodite.distributed import (
  21. get_tensor_model_parallel_world_size, )
  22. _VISION_MODEL_CLASSES = [
  23. LlavaForConditionalGeneration,
  24. ]
  25. @contextlib.contextmanager
  26. def _set_default_torch_dtype(dtype: torch.dtype):
  27. """Sets the default torch dtype to the given dtype."""
  28. old_dtype = torch.get_default_dtype()
  29. torch.set_default_dtype(dtype)
  30. yield
  31. torch.set_default_dtype(old_dtype)
  32. def _get_model_architecture(model_config: ModelConfig) -> Type[nn.Module]:
  33. architectures = getattr(model_config.hf_config, "architectures", [])
  34. for arch in architectures:
  35. model_cls = ModelRegistry.load_model_cls(arch)
  36. if model_cls is not None:
  37. return model_cls
  38. raise ValueError(
  39. f"Model architectures {architectures} are not supported for now. "
  40. f"Supported architectures: {ModelRegistry.get_supported_archs()}")
  41. def get_model(model_config: ModelConfig, device_config: DeviceConfig,
  42. **kwargs) -> nn.Module:
  43. lora_config = kwargs.get("lora_config", None)
  44. vision_language_config = kwargs.get("vision_language_config", None)
  45. model_class = _get_model_architecture(model_config)
  46. # Get the (maybe quantized) linear method.
  47. linear_method = None
  48. if model_config.quantization is not None:
  49. quant_config = get_quant_config(model_config)
  50. capability = torch.cuda.get_device_capability()
  51. capability = capability[0] * 10 + capability[1]
  52. if capability < quant_config.get_min_capability():
  53. raise ValueError(
  54. f"The quantization method {model_config.quantization} is not "
  55. "supported for the current GPU. "
  56. f"Minimum capability: {quant_config.get_min_capability()}. "
  57. f"Current capability: {capability}.")
  58. supported_dtypes = quant_config.get_supported_act_dtypes()
  59. if model_config.dtype not in supported_dtypes:
  60. # set the dtype to float16 for quantized models
  61. model_config.dtype = torch.float16
  62. logger.warning("Model is quantized. Forcing float16 datatype.")
  63. linear_method = quant_config.get_linear_method()
  64. with _set_default_torch_dtype(model_config.dtype):
  65. # Create a model instance.
  66. # The weights will be initialized as empty tensors.
  67. with torch.device(device_config.device) if not (
  68. isinstance(linear_method, BNBLinearMethod)
  69. and linear_method.quant_config.from_float) else nullcontext():
  70. if hasattr(model_class, "supported_lora_modules"):
  71. model = model_class(model_config.hf_config, linear_method,
  72. lora_config)
  73. elif lora_config:
  74. raise ValueError(
  75. f"Model {model_class.__name__} does not support LoRA, "
  76. "but LoRA is enabled. Support for this model may "
  77. "be added in the future. If this is important to you, "
  78. "please open an issue on github.")
  79. else:
  80. if model_class not in _VISION_MODEL_CLASSES:
  81. model = model_class(model_config.hf_config, linear_method)
  82. else:
  83. model = model_class(model_config.hf_config,
  84. vision_language_config, linear_method)
  85. if model_config.load_format == "dummy":
  86. # NOTE: For accurate performance evaluation, we assign
  87. # random values to the weights.
  88. initialize_dummy_weights(model)
  89. else:
  90. # Load the weights from the cached or downloaded files.
  91. model.load_weights(model_config.model, model_config.download_dir,
  92. model_config.load_format, model_config.revision)
  93. if isinstance(linear_method, BNBLinearMethod):
  94. replace_quant_params(
  95. model,
  96. quant_config=linear_method.quant_config,
  97. modules_to_not_convert="lm_head",
  98. )
  99. torch.cuda.synchronize()
  100. if linear_method.quant_config.from_float:
  101. model = model.cuda()
  102. gc.collect()
  103. torch.cuda.empty_cache()
  104. tp = get_tensor_model_parallel_world_size()
  105. logger.info(
  106. "Memory allocated for converted model: {} GiB x {} = {} "
  107. "GiB".format(
  108. round(
  109. torch.cuda.memory_allocated(torch.cuda.current_device()) /
  110. (1024 * 1024 * 1024),
  111. 2,
  112. ),
  113. tp,
  114. round(
  115. torch.cuda.memory_allocated(torch.cuda.current_device()) *
  116. tp / (1024 * 1024 * 1024),
  117. 2,
  118. ),
  119. ))
  120. logger.info(
  121. "Memory reserved for converted model: {} GiB x {} = {} "
  122. "GiB".format(
  123. round(
  124. torch.cuda.memory_reserved(torch.cuda.current_device()) /
  125. (1024 * 1024 * 1024),
  126. 2,
  127. ),
  128. tp,
  129. round(
  130. torch.cuda.memory_reserved(torch.cuda.current_device()) *
  131. tp / (1024 * 1024 * 1024),
  132. 2,
  133. ),
  134. ))
  135. return model.eval()