ray_gpu_executor.py 17 KB

123456789101112131415161718192021222324252627282930313233343536373839404142434445464748495051525354555657585960616263646566676869707172737475767778798081828384858687888990919293949596979899100101102103104105106107108109110111112113114115116117118119120121122123124125126127128129130131132133134135136137138139140141142143144145146147148149150151152153154155156157158159160161162163164165166167168169170171172173174175176177178179180181182183184185186187188189190191192193194195196197198199200201202203204205206207208209210211212213214215216217218219220221222223224225226227228229230231232233234235236237238239240241242243244245246247248249250251252253254255256257258259260261262263264265266267268269270271272273274275276277278279280281282283284285286287288289290291292293294295296297298299300301302303304305306307308309310311312313314315316317318319320321322323324325326327328329330331332333334335336337338339340341342343344345346347348349350351352353354355356357358359360361362363364365366367368369370371372373374375376377378379380381382383384385386387388389390391392393394395396397398399400401402403404405406407408409410411412413414415416417418419420421422423424425426427428429430431432433434435
  1. import asyncio
  2. import os
  3. import pickle
  4. from collections import defaultdict
  5. from itertools import islice, repeat
  6. from typing import TYPE_CHECKING, Any, Dict, List, Optional, Set, Tuple
  7. from loguru import logger
  8. from aphrodite.common.sequence import SamplerOutput, SequenceGroupMetadata
  9. from aphrodite.common.utils import (get_aphrodite_instance_id,
  10. get_distributed_init_method, get_ip,
  11. get_open_port, make_async)
  12. from aphrodite.executor.ray_utils import RayWorkerWrapper, ray
  13. from aphrodite.executor.executor_base import ExecutorAsyncBase, ExecutorBase
  14. from aphrodite.lora.request import LoRARequest
  15. if ray is not None:
  16. from ray.util.scheduling_strategies import PlacementGroupSchedulingStrategy
  17. if TYPE_CHECKING:
  18. from ray.util.placement_group import PlacementGroup
  19. # If the env var is set, it uses the Ray's compiled DAG API
  20. # which optimizes the control plane overhead.
  21. # Run Aphrodite with APHRODITE_USE_RAY_COMPILED_DAG=1 to enable it.
  22. USE_RAY_COMPILED_DAG = bool(os.getenv("APHRODITE_USE_RAY_COMPILED_DAG", 0))
  23. class RayGPUExecutor(ExecutorBase):
  24. def _init_executor(self) -> None:
  25. assert (not self.speculative_config
  26. ), "Speculative decoding not yet supported for RayGPU backend."
  27. assert self.parallel_config.worker_use_ray
  28. placement_group = self.parallel_config.placement_group
  29. # Disable Ray usage stats collection.
  30. ray_usage = os.environ.get("RAY_USAGE_STATS_ENABLED", "0")
  31. if ray_usage != "1":
  32. os.environ["RAY_USAGE_STATS_ENABLED"] = "0"
  33. # Create the parallel GPU workers.
  34. self._init_workers_ray(placement_group)
  35. self.forward_dag = None
  36. if USE_RAY_COMPILED_DAG:
  37. self.forward_dag = self._compiled_ray_dag()
  38. def _configure_ray_workers_use_nsight(self,
  39. ray_remote_kwargs) -> Dict[str, Any]:
  40. # If nsight profiling is enabled, we need to set the profiling
  41. # configuration for the ray workers as runtime env.
  42. runtime_env = ray_remote_kwargs.setdefault("runtime_env", {})
  43. runtime_env.update({
  44. "nsight": {
  45. "t": "cuda,cudnn,cublas",
  46. "o": "'worker_process_%p'",
  47. "cuda-graph-trace": "node",
  48. }
  49. })
  50. return ray_remote_kwargs
  51. def _init_workers_ray(self, placement_group: "PlacementGroup",
  52. **ray_remote_kwargs):
  53. if self.parallel_config.tensor_parallel_size == 1:
  54. # For single GPU case, we use a ray worker with constrained memory.
  55. num_gpus = self.cache_config.gpu_memory_utilization
  56. else:
  57. # Otherwise, the ray workers are allocated with a full GPU.
  58. num_gpus = 1
  59. # The driver dummy worker does not actually use any resources.
  60. # It holds the resource for the driver worker.
  61. self.driver_dummy_worker: RayWorkerWrapper = None
  62. # The remaining workers are the actual ray actors.
  63. self.workers: List[RayWorkerWrapper] = []
  64. if self.parallel_config.ray_workers_use_nsight:
  65. ray_remote_kwargs = self._configure_ray_workers_use_nsight(
  66. ray_remote_kwargs)
  67. # Create the workers.
  68. driver_ip = get_ip()
  69. for bundle_id, bundle in enumerate(placement_group.bundle_specs):
  70. if not bundle.get("GPU", 0):
  71. continue
  72. scheduling_strategy = PlacementGroupSchedulingStrategy(
  73. placement_group=placement_group,
  74. placement_group_capture_child_tasks=True,
  75. placement_group_bundle_index=bundle_id,
  76. )
  77. worker = ray.remote(
  78. num_cpus=0,
  79. num_gpus=num_gpus,
  80. scheduling_strategy=scheduling_strategy,
  81. **ray_remote_kwargs,
  82. )(RayWorkerWrapper).remote(
  83. worker_module_name="aphrodite.task_handler.worker",
  84. worker_class_name="Worker",
  85. trust_remote_code=self.model_config.trust_remote_code,
  86. )
  87. worker_ip = ray.get(worker.get_node_ip.remote())
  88. if worker_ip == driver_ip and self.driver_dummy_worker is None:
  89. # If the worker is on the same node as the driver, we use it
  90. # as the resource holder for the driver process.
  91. self.driver_dummy_worker = worker
  92. self.driver_worker = RayWorkerWrapper(
  93. worker_module_name="aphrodite.task_handler.worker",
  94. worker_class_name="Worker",
  95. trust_remote_code=self.model_config.trust_remote_code,
  96. )
  97. else:
  98. # Else, added to the list of workers.
  99. self.workers.append(worker)
  100. if self.driver_dummy_worker is None:
  101. raise ValueError(
  102. "Ray does not allocate any GPUs on the driver node. Consider "
  103. "adjusting the Ray placement group or running the driver on a "
  104. "GPU node.")
  105. # Get the set of GPU IDs used on each node.
  106. worker_node_and_gpu_ids = self._run_workers("get_node_and_gpu_ids",
  107. use_dummy_driver=True)
  108. node_workers = defaultdict(list)
  109. node_gpus = defaultdict(list)
  110. for i, (node_id, gpu_ids) in enumerate(worker_node_and_gpu_ids):
  111. node_workers[node_id].append(i)
  112. node_gpus[node_id].extend(gpu_ids)
  113. for node_id, gpu_ids in node_gpus.items():
  114. node_gpus[node_id] = sorted(gpu_ids)
  115. APHRODITE_INSTANCE_ID = get_aphrodite_instance_id()
  116. # Set environment variables for the driver and workers.
  117. all_args_to_update_environment_variables = [({
  118. "CUDA_VISIBLE_DEVICES":
  119. ",".join(map(str, node_gpus[node_id])),
  120. "APHRODITE_INSTANCE_ID":
  121. APHRODITE_INSTANCE_ID,
  122. "APHRODITE_TRACE_FUNCTION":
  123. os.getenv("APHRODITE_TRACE_FUNCTION", "0"),
  124. }, ) for (node_id, _) in worker_node_and_gpu_ids]
  125. self._run_workers("update_environment_variables",
  126. all_args=all_args_to_update_environment_variables)
  127. distributed_init_method = get_distributed_init_method(
  128. driver_ip, get_open_port())
  129. def collect_arg_helper_func(**kwargs):
  130. # avoid writing `{"name": value}` manually
  131. return kwargs
  132. # Initialize the actual workers inside worker wrapper.
  133. init_worker_all_kwargs = []
  134. for rank, (node_id, _) in enumerate(worker_node_and_gpu_ids):
  135. local_rank = node_workers[node_id].index(rank)
  136. init_worker_all_kwargs.append(
  137. collect_arg_helper_func(
  138. model_config=self.model_config,
  139. parallel_config=self.parallel_config,
  140. scheduler_config=self.scheduler_config,
  141. device_config=self.device_config,
  142. cache_config=self.cache_config,
  143. load_config=self.load_config,
  144. local_rank=local_rank,
  145. rank=rank,
  146. distributed_init_method=distributed_init_method,
  147. lora_config=self.lora_config,
  148. vision_language_config=self.vision_language_config,
  149. is_driver_worker=rank == 0,
  150. ))
  151. self._run_workers("init_worker", all_kwargs=init_worker_all_kwargs)
  152. self._run_workers("init_device")
  153. self._run_workers(
  154. "load_model",
  155. max_concurrent_workers=self.parallel_config.
  156. max_parallel_loading_workers,
  157. )
  158. def determine_num_available_blocks(self) -> Tuple[int, int]:
  159. """Determine the number of available KV blocks.
  160. This invokes `determine_num_available_blocks` on each worker and takes
  161. the min of the results, guaranteeing that the selected cache sizes are
  162. compatible with all workers.
  163. Returns:
  164. - Tuple[num_gpu_blocks, num_cpu_blocks]
  165. """
  166. # Get the maximum number of blocks that can be allocated on GPU and CPU.
  167. num_blocks = self._run_workers("determine_num_available_blocks", )
  168. # Since we use a shared centralized controller, we take the minimum
  169. # number of blocks across all workers to make sure all the memory
  170. # operators can be applied to all workers.
  171. num_gpu_blocks = min(b[0] for b in num_blocks)
  172. num_cpu_blocks = min(b[1] for b in num_blocks)
  173. return num_gpu_blocks, num_cpu_blocks
  174. def initialize_cache(self, num_gpu_blocks: int,
  175. num_cpu_blocks: int) -> None:
  176. """Initialize the KV cache in all workers.
  177. """
  178. # NOTE: We log here to avoid multiple logs when number of workers is
  179. # greater than one. We could log in the engine, but not all executors
  180. # have GPUs.
  181. logger.info(f"# GPU blocks: {num_gpu_blocks}, "
  182. f"# CPU blocks: {num_cpu_blocks}")
  183. logger.info(
  184. f"Minimum concurrency: {num_gpu_blocks * self.cache_config.block_size / self.scheduler_config.max_model_len:.2f}x" # noqa: E501
  185. )
  186. self.cache_config.num_gpu_blocks = num_gpu_blocks
  187. self.cache_config.num_cpu_blocks = num_cpu_blocks
  188. self._run_workers("initialize_cache",
  189. num_gpu_blocks=num_gpu_blocks,
  190. num_cpu_blocks=num_cpu_blocks)
  191. def execute_model(self,
  192. seq_group_metadata_list: List[SequenceGroupMetadata],
  193. blocks_to_swap_in: Dict[int, int],
  194. blocks_to_swap_out: Dict[int, int],
  195. blocks_to_copy: Dict[int, List[int]],
  196. num_lookahead_slots: int = 0) -> SamplerOutput:
  197. all_outputs = self._run_workers(
  198. "execute_model",
  199. driver_kwargs={
  200. "seq_group_metadata_list": seq_group_metadata_list,
  201. "blocks_to_swap_in": blocks_to_swap_in,
  202. "blocks_to_swap_out": blocks_to_swap_out,
  203. "blocks_to_copy": blocks_to_copy,
  204. },
  205. use_ray_compiled_dag=USE_RAY_COMPILED_DAG)
  206. # Only the driver worker returns the sampling results.
  207. output = all_outputs[0]
  208. return output
  209. def add_lora(self, lora_request: LoRARequest) -> bool:
  210. assert lora_request.lora_int_id > 0, "lora_id must be greater than 0."
  211. return self._run_workers(
  212. "add_lora",
  213. lora_request=lora_request,
  214. )
  215. def remove_lora(self, lora_id: int) -> bool:
  216. assert lora_id > 0, "lora_id must be greater than 0."
  217. return self._run_workers(
  218. "remove_lora",
  219. lora_id=lora_id,
  220. )
  221. def list_loras(self) -> Set[int]:
  222. return self._run_workers("list_loras")
  223. def _run_workers(
  224. self,
  225. method: str,
  226. *args,
  227. driver_args: Optional[Tuple[Any, ...]] = None,
  228. driver_kwargs: Optional[Dict[str, Any]] = None,
  229. all_args: Optional[List[Tuple[Any, ...]]] = None,
  230. all_kwargs: Optional[List[Dict[str, Any]]] = None,
  231. use_dummy_driver: bool = False,
  232. max_concurrent_workers: Optional[int] = None,
  233. use_ray_compiled_dag: bool = False,
  234. **kwargs,
  235. ) -> Any:
  236. """Runs the given method on all workers. Can be used in the following
  237. ways:
  238. - args/kwargs: All workers share the same args/kwargs
  239. - args/kwargs and driver_args/driver_kwargs: Driver worker has
  240. different args
  241. - all_args/all_kwargs: args/kwargs for each worker are specified
  242. individually
  243. """
  244. if max_concurrent_workers:
  245. raise NotImplementedError(
  246. "max_concurrent_workers is not supported yet.")
  247. if driver_args is None:
  248. driver_args = args if all_args is None else all_args[0]
  249. if driver_kwargs is None:
  250. driver_kwargs = kwargs if all_kwargs is None else all_kwargs[0]
  251. count = len(self.workers)
  252. all_worker_args = repeat(args, count) if all_args is None \
  253. else islice(all_args, 1, None)
  254. all_worker_kwargs = repeat(kwargs, count) if all_kwargs is None \
  255. else islice(all_kwargs, 1, None)
  256. if use_ray_compiled_dag:
  257. # Right now, compiled DAG can only accept a single
  258. # input. TODO: Fix it.
  259. assert self.forward_dag is not None
  260. output_channels = self.forward_dag.execute(1)
  261. else:
  262. # Start the ray workers first.
  263. ray_worker_outputs = [
  264. worker.execute_method.remote(method, *worker_args,
  265. **worker_kwargs)
  266. for (worker, worker_args, worker_kwargs
  267. ) in zip(self.workers, all_worker_args, all_worker_kwargs)
  268. ]
  269. # Start the driver worker after all the ray workers.
  270. if not use_dummy_driver:
  271. driver_worker_output = self.driver_worker.execute_method(
  272. method, *driver_args, **driver_kwargs)
  273. else:
  274. driver_worker_output = ray.get(
  275. self.driver_dummy_worker.execute_method.remote(
  276. method, *driver_args, **driver_kwargs))
  277. # Get the results of the ray workers.
  278. if self.workers:
  279. if use_ray_compiled_dag:
  280. try:
  281. ray_worker_outputs = [
  282. pickle.loads(chan.begin_read())
  283. for chan in output_channels
  284. ]
  285. finally:
  286. # Has to call end_read in order to reuse the DAG.
  287. for chan in output_channels:
  288. chan.end_read()
  289. else:
  290. ray_worker_outputs = ray.get(ray_worker_outputs)
  291. return [driver_worker_output] + ray_worker_outputs
  292. def _compiled_ray_dag(self):
  293. import pkg_resources
  294. required_version = "2.9"
  295. current_version = pkg_resources.get_distribution("ray").version
  296. if current_version < required_version:
  297. raise ValueError(f"Ray version {required_version} or greater is "
  298. f"required, but found {current_version}")
  299. from ray.dag import InputNode, MultiOutputNode
  300. assert self.parallel_config.worker_use_ray
  301. # Right now, compiled DAG requires at least 1 arg. We send
  302. # a dummy value for now. It will be fixed soon.
  303. with InputNode() as input_data:
  304. forward_dag = MultiOutputNode([
  305. worker.execute_model_compiled_dag_remote.bind(input_data)
  306. for worker in self.workers
  307. ])
  308. return forward_dag.experimental_compile()
  309. def check_health(self) -> None:
  310. """Raises an error if engine is unhealthy."""
  311. self._check_if_any_actor_is_dead()
  312. def _check_if_any_actor_is_dead(self):
  313. if not self.workers:
  314. return
  315. dead_actors = []
  316. for actor in self.workers:
  317. actor_state = ray.state.actors(actor._ray_actor_id.hex()) # pylint: disable=protected-access
  318. if actor_state["State"] == "DEAD":
  319. dead_actors.append(actor)
  320. if dead_actors:
  321. raise RuntimeError("At least one Worker is dead. "
  322. f"Dead Workers: {dead_actors}. ")
  323. class RayGPUExecutorAsync(RayGPUExecutor, ExecutorAsyncBase):
  324. def __init__(self, *args, **kwargs):
  325. super().__init__(*args, **kwargs)
  326. self.driver_executor = make_async(self.driver_worker.execute_method)
  327. async def _run_workers_async(
  328. self,
  329. method: str,
  330. *args,
  331. driver_args: Optional[Tuple[Any, ...]] = None,
  332. driver_kwargs: Optional[Dict[str, Any]] = None,
  333. **kwargs,
  334. ) -> Any:
  335. """Runs the given method on all workers."""
  336. coros = []
  337. if driver_args is None:
  338. driver_args = args
  339. if driver_kwargs is None:
  340. driver_kwargs = kwargs
  341. coros.append(
  342. self.driver_executor(method, *driver_args, **driver_kwargs))
  343. # Run the ray workers asynchronously.
  344. for worker in self.workers:
  345. coros.append(worker.execute_method.remote(method, *args, **kwargs))
  346. all_outputs = await asyncio.gather(*coros)
  347. return all_outputs
  348. async def execute_model_async(
  349. self,
  350. seq_group_metadata_list: List[SequenceGroupMetadata],
  351. blocks_to_swap_in: Dict[int, int],
  352. blocks_to_swap_out: Dict[int, int],
  353. blocks_to_copy: Dict[int, List[int]],
  354. num_lookahead_slots: int,
  355. ) -> SamplerOutput:
  356. all_outputs = await self._run_workers_async(
  357. "execute_model",
  358. driver_kwargs={
  359. "seq_group_metadata_list": seq_group_metadata_list,
  360. "blocks_to_swap_in": blocks_to_swap_in,
  361. "blocks_to_swap_out": blocks_to_swap_out,
  362. "blocks_to_copy": blocks_to_copy,
  363. "num_lookahead_slots": num_lookahead_slots,
  364. })
  365. # Only the driver worker returns the sampling results.
  366. output = all_outputs[0]
  367. return output