worker.py 14 KB

123456789101112131415161718192021222324252627282930313233343536373839404142434445464748495051525354555657585960616263646566676869707172737475767778798081828384858687888990919293949596979899100101102103104105106107108109110111112113114115116117118119120121122123124125126127128129130131132133134135136137138139140141142143144145146147148149150151152153154155156157158159160161162163164165166167168169170171172173174175176177178179180181182183184185186187188189190191192193194195196197198199200201202203204205206207208209210211212213214215216217218219220221222223224225226227228229230231232233234235236237238239240241242243244245246247248249250251252253254255256257258259260261262263264265266267268269270271272273274275276277278279280281282283284285286287288289290291292293294295296297298299300301302303304305306307308309310311312313314315316317318319320321322323324325326327328329330331332333334335336337338339340341342343344345346
  1. """A GPU worker class."""
  2. import gc
  3. import os
  4. from typing import Any, Dict, List, Optional, Set, Tuple
  5. import torch
  6. import torch.distributed
  7. from loguru import logger
  8. from aphrodite.common.sequence import SamplerOutput, SequenceGroupMetadata
  9. from aphrodite.common.config import (CacheConfig, DeviceConfig, LoadConfig,
  10. LoRAConfig, ModelConfig, ParallelConfig,
  11. SchedulerConfig, VisionLanguageConfig)
  12. from aphrodite.distributed import (broadcast_tensor_dict,
  13. ensure_model_parallel_initialized,
  14. init_distributed_environment)
  15. from aphrodite.distributed.device_communicators import pynccl_utils
  16. from aphrodite.distributed.device_communicators.custom_all_reduce import \
  17. init_custom_ar
  18. from aphrodite.lora.request import LoRARequest
  19. from aphrodite.modeling import set_random_seed
  20. from aphrodite.task_handler.cache_engine import CacheEngine
  21. from aphrodite.task_handler.model_runner import ModelRunner
  22. from aphrodite.task_handler.worker_base import WorkerBase
  23. class Worker(WorkerBase):
  24. """A worker class that executes (a partition of) the model on a GPU.
  25. Each worker is associated with a single GPU. The worker is responsible for
  26. maintaining the KV cache and executing the model on the GPU. In case of
  27. distributed inference, each worker is assigned a partition of the model.
  28. """
  29. def __init__(
  30. self,
  31. model_config: ModelConfig,
  32. parallel_config: ParallelConfig,
  33. scheduler_config: SchedulerConfig,
  34. device_config: DeviceConfig,
  35. cache_config: CacheConfig,
  36. load_config: LoadConfig,
  37. local_rank: int,
  38. rank: int,
  39. distributed_init_method: str,
  40. lora_config: Optional[LoRAConfig] = None,
  41. vision_language_config: Optional[VisionLanguageConfig] = None,
  42. is_driver_worker: bool = False,
  43. ) -> None:
  44. self.model_config = model_config
  45. self.parallel_config = parallel_config
  46. self.scheduler_config = scheduler_config
  47. self.device_config = device_config
  48. self.cache_config = cache_config
  49. self.local_rank = local_rank
  50. self.rank = rank
  51. self.distributed_init_method = distributed_init_method
  52. self.lora_config = lora_config
  53. self.load_config = load_config
  54. self.is_driver_worker = is_driver_worker
  55. if self.is_driver_worker:
  56. assert self.rank == 0, "The driver worker must have rank 0."
  57. if self.model_config.trust_remote_code:
  58. # note: lazy import to avoid importing torch before initializing
  59. from aphrodite.common.utils import init_cached_hf_modules
  60. init_cached_hf_modules()
  61. self.vision_language_config = vision_language_config
  62. if self.vision_language_config:
  63. assert not self.lora_config, (
  64. "To be tested: vision language model with LoRA settings.")
  65. self.model_runner = ModelRunner(
  66. model_config,
  67. parallel_config,
  68. scheduler_config,
  69. device_config,
  70. load_config=load_config,
  71. lora_config=self.lora_config,
  72. kv_cache_dtype=self.cache_config.cache_dtype,
  73. is_driver_worker=is_driver_worker,
  74. vision_language_config=vision_language_config,
  75. )
  76. # Uninitialized cache engine. Will be initialized by
  77. # initialize_cache.
  78. self.cache_engine: CacheEngine
  79. self.gpu_cache: List[torch.Tensor]
  80. def init_device(self) -> None:
  81. if self.device_config.device.type == "cuda":
  82. # torch.distributed.all_reduce does not free the input tensor until
  83. # the synchronization point. This causes the memory usage to grow
  84. # as the number of all_reduce calls increases. This env var disables
  85. # this behavior.
  86. # Related issue:
  87. # https://discuss.pytorch.org/t/cuda-allocation-lifetime-for-inputs-to-distributed-all-reduce/191573
  88. os.environ["TORCH_NCCL_AVOID_RECORD_STREAMS"] = "1"
  89. # This env var set by Ray causes exceptions with graph building.
  90. os.environ.pop("NCCL_ASYNC_ERROR_HANDLING", None)
  91. self.device = torch.device(f"cuda:{self.local_rank}")
  92. torch.cuda.set_device(self.device)
  93. _check_if_gpu_supports_dtype(self.model_config.dtype)
  94. torch.cuda.empty_cache()
  95. self.init_gpu_memory = torch.cuda.mem_get_info()[0]
  96. else:
  97. raise RuntimeError(
  98. f"Not support device type: {self.device_config.device}")
  99. # Initialize the distributed environment.
  100. init_worker_distributed_environment(self.parallel_config, self.rank,
  101. self.distributed_init_method,
  102. self.local_rank)
  103. # Set random seed.
  104. set_random_seed(self.model_config.seed)
  105. def load_model(self):
  106. self.model_runner.load_model()
  107. @torch.inference_mode()
  108. def determine_num_available_blocks(self) -> Tuple[int, int]:
  109. """Profiles the peak memory usage of the model to determine how many
  110. KV blocks may be allocated without OOMs.
  111. The engine will first conduct a profiling of the existing memory usage.
  112. Then, it calculate the maximum possible number of GPU and CPU blocks
  113. that can be allocated with the remaining free memory.
  114. .. tip::
  115. You may limit the usage of GPU memory
  116. by adjusting the `gpu_memory_utilization` parameter.
  117. """
  118. # Profile the memory usage of the model and get the maximum number of
  119. # cache blocks that can be allocated with the remaining free memory.
  120. torch.cuda.empty_cache()
  121. # Execute a forward pass with dummy inputs to profile the memory usage
  122. # of the model.
  123. self.model_runner.profile_run()
  124. # Calculate the number of blocks that can be allocated with the
  125. # profiled peak memory.
  126. torch.cuda.synchronize()
  127. free_gpu_memory, total_gpu_memory = torch.cuda.mem_get_info()
  128. # NOTE: Here we assume that the other processes using the same
  129. # GPU did not change their memory usage during the profiling.
  130. peak_memory = self.init_gpu_memory - free_gpu_memory
  131. assert peak_memory > 0, (
  132. "Error in memory profiling. This happens when the GPU memory was "
  133. "not properly cleaned up before initializing Aphrodite.")
  134. cache_block_size = self.get_cache_block_size_bytes()
  135. num_gpu_blocks = int(
  136. (total_gpu_memory * self.cache_config.gpu_memory_utilization -
  137. peak_memory) // cache_block_size)
  138. num_cpu_blocks = int(self.cache_config.swap_space_bytes //
  139. cache_block_size)
  140. num_gpu_blocks = max(num_gpu_blocks, 0)
  141. num_cpu_blocks = max(num_cpu_blocks, 0)
  142. if self.model_runner.lora_manager:
  143. self.model_runner.remove_all_loras()
  144. gc.collect()
  145. torch.cuda.empty_cache()
  146. return num_gpu_blocks, num_cpu_blocks
  147. def initialize_cache(self, num_gpu_blocks: int,
  148. num_cpu_blocks: int) -> None:
  149. """Allocate GPU and CPU KV cache with the specified number of blocks.
  150. This also warms up the model, which may record CUDA graphs.
  151. """
  152. raise_if_cache_size_invalid(num_gpu_blocks,
  153. self.cache_config.block_size,
  154. self.model_config.max_model_len)
  155. self.cache_config.num_gpu_blocks = num_gpu_blocks
  156. self.cache_config.num_cpu_blocks = num_cpu_blocks
  157. self._init_cache_engine()
  158. self._warm_up_model()
  159. def _init_cache_engine(self):
  160. assert self.cache_config.num_gpu_blocks is not None
  161. self.cache_engine = CacheEngine(self.cache_config, self.model_config,
  162. self.parallel_config)
  163. self.gpu_cache = self.cache_engine.gpu_cache
  164. self.model_runner.set_block_size(self.cache_engine.block_size)
  165. def _warm_up_model(self) -> None:
  166. if not self.model_config.enforce_eager:
  167. self.model_runner.capture_model(self.gpu_cache)
  168. # Reset the seed to ensure that the random state is not affected by
  169. # the model initialization and profiling.
  170. set_random_seed(self.model_config.seed)
  171. def cache_swap(
  172. self,
  173. blocks_to_swap_in: Dict[int, int],
  174. blocks_to_swap_out: Dict[int, int],
  175. blocks_to_copy: Dict[int, List[int]],
  176. ) -> None:
  177. # Issue cache operations.
  178. # TODO: Profile swapping overhead and optimize if needed.
  179. if blocks_to_swap_in:
  180. self.cache_engine.swap_in(blocks_to_swap_in)
  181. if blocks_to_swap_out:
  182. self.cache_engine.swap_out(blocks_to_swap_out)
  183. if blocks_to_copy:
  184. self.cache_engine.copy(blocks_to_copy)
  185. @torch.inference_mode()
  186. def execute_model(
  187. self,
  188. seq_group_metadata_list: Optional[List[SequenceGroupMetadata]] = None,
  189. blocks_to_swap_in: Optional[Dict[int, int]] = None,
  190. blocks_to_swap_out: Optional[Dict[int, int]] = None,
  191. blocks_to_copy: Optional[Dict[int, List[int]]] = None,
  192. num_lookahead_slots: int = 0,
  193. ) -> List[SamplerOutput]:
  194. if self.is_driver_worker:
  195. assert seq_group_metadata_list is not None
  196. num_seq_groups = len(seq_group_metadata_list)
  197. assert blocks_to_swap_in is not None
  198. assert blocks_to_swap_out is not None
  199. assert blocks_to_copy is not None
  200. data: Dict[str, Any] = {
  201. "num_seq_groups": num_seq_groups,
  202. "blocks_to_swap_in": blocks_to_swap_in,
  203. "blocks_to_swap_out": blocks_to_swap_out,
  204. "blocks_to_copy": blocks_to_copy,
  205. }
  206. broadcast_tensor_dict(data, src=0)
  207. else:
  208. data = broadcast_tensor_dict(src=0)
  209. num_seq_groups = data["num_seq_groups"]
  210. blocks_to_swap_in = data["blocks_to_swap_in"]
  211. blocks_to_swap_out = data["blocks_to_swap_out"]
  212. blocks_to_copy = data["blocks_to_copy"]
  213. assert blocks_to_swap_in is not None
  214. assert blocks_to_swap_out is not None
  215. assert blocks_to_copy is not None
  216. self.cache_swap(blocks_to_swap_in, blocks_to_swap_out, blocks_to_copy)
  217. # If there is no input, we don't need to execute the model.
  218. if num_seq_groups == 0:
  219. return []
  220. output = self.model_runner.execute_model(seq_group_metadata_list,
  221. self.gpu_cache)
  222. # Worker only supports single-step execution. Wrap the output in a list
  223. # to conform to interface.
  224. return [output]
  225. def add_lora(self, lora_request: LoRARequest) -> bool:
  226. return self.model_runner.add_lora(lora_request)
  227. def remove_lora(self, lora_id: int) -> bool:
  228. return self.model_runner.remove_lora(lora_id)
  229. def list_loras(self) -> Set[int]:
  230. return self.model_runner.list_loras()
  231. @property
  232. def max_model_len(self) -> int:
  233. return self.model_config.max_model_len
  234. @property
  235. def vocab_size(self) -> int:
  236. return self.model_runner.vocab_size
  237. def get_cache_block_size_bytes(self) -> int:
  238. """Get the size of the KV cache block size in bytes.
  239. """
  240. return CacheEngine.get_cache_block_size(self.cache_config,
  241. self.model_config,
  242. self.parallel_config)
  243. def init_worker_distributed_environment(
  244. parallel_config: ParallelConfig,
  245. rank: int,
  246. distributed_init_method: Optional[str] = None,
  247. local_rank: int = -1,
  248. ) -> None:
  249. """Initialize the distributed environment."""
  250. init_distributed_environment(parallel_config.world_size, rank,
  251. distributed_init_method, local_rank)
  252. if pynccl_utils.is_initialized():
  253. pynccl_world_size = pynccl_utils.get_world_size()
  254. if pynccl_world_size != parallel_config.world_size:
  255. raise RuntimeError(
  256. "pynccl is already initialized but the pynccl world "
  257. "size does not match parallel_config.world_size "
  258. f"({pynccl_world_size} vs. {parallel_config.world_size}).")
  259. elif parallel_config.world_size > 1:
  260. # NOTE: We don't initialize pynccl process group when world size
  261. # is 1.
  262. pynccl_utils.init_process_group()
  263. ensure_model_parallel_initialized(parallel_config.tensor_parallel_size,
  264. parallel_config.pipeline_parallel_size)
  265. # Initialize a custom fast all-reduce implementation.
  266. if not parallel_config.disable_custom_all_reduce:
  267. init_custom_ar()
  268. # A small all_reduce for warmup.
  269. torch.distributed.all_reduce(torch.zeros(1).cuda())
  270. if pynccl_utils.is_initialized():
  271. pynccl_utils.all_reduce(torch.zeros(1).cuda())
  272. def _check_if_gpu_supports_dtype(torch_dtype: torch.dtype):
  273. # Check if the GPU supports the dtype.
  274. if torch_dtype == torch.bfloat16:
  275. compute_capability = torch.cuda.get_device_capability()
  276. if compute_capability[0] < 8:
  277. gpu_name = torch.cuda.get_device_name()
  278. raise ValueError(
  279. "Bfloat16 is only supported on GPUs with compute capability "
  280. f"of at least 8.0. Your {gpu_name} GPU has compute capability "
  281. f"{compute_capability[0]}.{compute_capability[1]}. "
  282. "You can use float16 instead by explicitly setting the"
  283. "`dtype` flag in CLI, for example: --dtype=half.")
  284. def raise_if_cache_size_invalid(num_gpu_blocks, block_size,
  285. max_model_len) -> None:
  286. if num_gpu_blocks <= 0:
  287. raise ValueError("No available memory for the cache blocks. "
  288. "Try increasing `gpu_memory_utilization` when "
  289. "initializing the engine.")
  290. max_seq_len = block_size * num_gpu_blocks
  291. logger.info(f"Maximum sequence length allowed in the cache: "
  292. f"{max_seq_len}")
  293. if max_model_len > max_seq_len:
  294. raise ValueError(
  295. f"The model's max seq len ({max_model_len}) "
  296. "is larger than the maximum number of tokens that can be "
  297. f"stored in KV cache ({max_seq_len}). Try increasing "
  298. "`gpu_memory_utilization` or decreasing `max_model_len` when "
  299. "initializing the engine.")