parallel_state.py 12 KB

123456789101112131415161718192021222324252627282930313233343536373839404142434445464748495051525354555657585960616263646566676869707172737475767778798081828384858687888990919293949596979899100101102103104105106107108109110111112113114115116117118119120121122123124125126127128129130131132133134135136137138139140141142143144145146147148149150151152153154155156157158159160161162163164165166167168169170171172173174175176177178179180181182183184185186187188189190191192193194195196197198199200201202203204205206207208209210211212213214215216217218219220221222223224225226227228229230231232233234235236237238239240241242243244245246247248249250251252253254255256257258259260261262263264265266267268269270271272273274275276277278279280281282283284285286287288289290291292293294295296297298299300301302303304305306307308309310311312313314
  1. # Copyright 2023 The PygmalionAI team.
  2. # Copyright 2023 The vLLM team.
  3. # Adapted from
  4. # https://github.com/NVIDIA/Megatron-LM/blob/main/megatron/core/parallel_state.py
  5. # Copyright (c) 2022, NVIDIA CORPORATION. All rights reserved.
  6. """Tensor and pipeline parallel groups."""
  7. import contextlib
  8. from typing import Optional
  9. import torch
  10. from loguru import logger
  11. # Tensor model parallel group that the current rank belongs to.
  12. _TENSOR_MODEL_PARALLEL_GROUP = None
  13. # Pipeline model parallel group that the current rank belongs to.
  14. _PIPELINE_MODEL_PARALLEL_GROUP = None
  15. # when people blindly call `torch.distributed.all_reduce` etc,
  16. # it will use this group. It is initialized with the `backend`
  17. # parameter of `init_distributed_environment` below.
  18. # Essentially, this is `torch.distributed.group.WORLD`.
  19. # We leave a line here to note that this is device-specific.
  20. # Note that this variable is not safe to use, because when users
  21. # call `init_distributed_environment` first, and then destroy
  22. # the process group themselves, this variable will keep a reference to the
  23. # destroyed process group, which is not useful.
  24. _DEVICE_WORLD_GROUP = None
  25. # duing `init_distributed_environment`, we will also initialize a
  26. # group with `gloo` backend, to allow direct coordination between
  27. # processes through the CPU.
  28. _CPU_WORLD_GROUP = None
  29. # In summary, after calling `init_distributed_environment`, we will
  30. # always have two groups: one for device-specific (and is the default)
  31. # and one for CPU. All processes will be part of both groups.
  32. # A list of global ranks for each pipeline group to ease calculation of the
  33. # source rank when broadcasting from the first or last pipeline stage.
  34. _PIPELINE_GLOBAL_RANKS = None
  35. _LOCAL_RANK = -1
  36. def get_local_rank():
  37. global _LOCAL_RANK
  38. return _LOCAL_RANK
  39. def init_distributed_environment(
  40. world_size: int = -1,
  41. rank: int = -1,
  42. distributed_init_method: str = "env://",
  43. local_rank: int = -1,
  44. backend: str = "nccl",
  45. ):
  46. logger.debug(f"{world_size=} {rank=} {local_rank=} "
  47. f"{distributed_init_method=} {backend=}")
  48. if not torch.distributed.is_initialized():
  49. assert distributed_init_method is not None, (
  50. "distributed_init_method must be provided when initializing "
  51. "distributed environment")
  52. # this backend is used for WORLD
  53. torch.distributed.init_process_group(
  54. backend=backend,
  55. init_method=distributed_init_method,
  56. world_size=world_size,
  57. rank=rank)
  58. global _DEVICE_WORLD_GROUP, _CPU_WORLD_GROUP
  59. _DEVICE_WORLD_GROUP = torch.distributed.group.WORLD
  60. ranks = list(range(torch.distributed.get_world_size()))
  61. _CPU_WORLD_GROUP = torch.distributed.new_group(ranks=ranks,
  62. backend="gloo")
  63. global _LOCAL_RANK
  64. _LOCAL_RANK = local_rank
  65. def initialize_model_parallel(
  66. tensor_model_parallel_size: int = 1,
  67. pipeline_model_parallel_size: int = 1,
  68. backend: Optional[str] = None,
  69. ) -> None:
  70. """
  71. Initialize model parallel groups.
  72. Arguments:
  73. tensor_model_parallel_size: number of GPUs used for tensor model
  74. parallelism.
  75. pipeline_model_parallel_size: number of GPUs used for pipeline model
  76. parallelism.
  77. Let's say we have a total of 8 GPUs denoted by g0 ... g7 and we
  78. use 2 GPUs to parallelize the model tensor, and 4 GPUs to parallelize
  79. the model pipeline. The present function will
  80. create 4 tensor model-parallel groups and 2 pipeline model-parallel groups:
  81. 4 tensor model-parallel groups:
  82. [g0, g1], [g2, g3], [g4, g5], [g6, g7]
  83. 2 pipeline model-parallel groups:
  84. [g0, g2, g4, g6], [g1, g3, g5, g7]
  85. Note that for efficiency, the caller should make sure adjacent ranks
  86. are on the same DGX box. For example if we are using 2 DGX-1 boxes
  87. with a total of 16 GPUs, rank 0 to 7 belong to the first box and
  88. ranks 8 to 15 belong to the second box.
  89. """
  90. # Get world size and rank. Ensure some consistencies.
  91. assert torch.distributed.is_initialized()
  92. world_size: int = torch.distributed.get_world_size()
  93. # get the backend of _DEVICE_WORLD_GROUP
  94. backend = backend or torch.distributed.get_backend()
  95. if (world_size !=
  96. tensor_model_parallel_size * pipeline_model_parallel_size):
  97. raise RuntimeError(
  98. f"world_size ({world_size}) is not equal to "
  99. f"tensor_model_parallel_size ({tensor_model_parallel_size}) x "
  100. f"pipeline_model_parallel_size ({pipeline_model_parallel_size})")
  101. num_tensor_model_parallel_groups: int = (world_size //
  102. tensor_model_parallel_size)
  103. num_pipeline_model_parallel_groups: int = (world_size //
  104. pipeline_model_parallel_size)
  105. rank = torch.distributed.get_rank()
  106. # Build the tensor model-parallel groups.
  107. global _TENSOR_MODEL_PARALLEL_GROUP
  108. assert _TENSOR_MODEL_PARALLEL_GROUP is None, (
  109. "tensor model parallel group is already initialized")
  110. for i in range(num_tensor_model_parallel_groups):
  111. ranks = range(i * tensor_model_parallel_size,
  112. (i + 1) * tensor_model_parallel_size)
  113. group = torch.distributed.new_group(ranks, backend=backend)
  114. if rank in ranks:
  115. _TENSOR_MODEL_PARALLEL_GROUP = group
  116. # Build the pipeline model-parallel groups.
  117. global _PIPELINE_MODEL_PARALLEL_GROUP
  118. global _PIPELINE_GLOBAL_RANKS
  119. assert _PIPELINE_MODEL_PARALLEL_GROUP is None, (
  120. "pipeline model parallel group is already initialized")
  121. for i in range(num_pipeline_model_parallel_groups):
  122. ranks = range(i, world_size, num_pipeline_model_parallel_groups)
  123. group = torch.distributed.new_group(ranks, backend=backend)
  124. if rank in ranks:
  125. _PIPELINE_MODEL_PARALLEL_GROUP = group
  126. _PIPELINE_GLOBAL_RANKS = ranks
  127. def ensure_model_parallel_initialized(
  128. tensor_model_parallel_size: int,
  129. pipeline_model_parallel_size: int,
  130. backend: Optional[str] = None,
  131. ) -> None:
  132. """Helper to initialize model parallel groups if they are not initialized,
  133. or ensure tensor-parallel and pipeline-parallel sizes are equal to expected
  134. values if the model parallel groups are initialized.
  135. """
  136. # get the backend of _DEVICE_WORLD_GROUP
  137. backend = backend or torch.distributed.get_backend()
  138. if not model_parallel_is_initialized():
  139. initialize_model_parallel(tensor_model_parallel_size,
  140. pipeline_model_parallel_size, backend)
  141. return
  142. assert (
  143. get_tensor_model_parallel_world_size() == tensor_model_parallel_size
  144. ), ("tensor parallel group already initialized, but of unexpected size: "
  145. f"{get_tensor_model_parallel_world_size()=} vs. "
  146. f"{tensor_model_parallel_size=}")
  147. assert (get_pipeline_model_parallel_world_size(
  148. ) == pipeline_model_parallel_size), (
  149. "pipeline parallel group already initialized, but of unexpected size: "
  150. f"{get_pipeline_model_parallel_world_size()=} vs. "
  151. f"{pipeline_model_parallel_size=}")
  152. def model_parallel_is_initialized():
  153. """Check if tensor and pipeline parallel groups are initialized."""
  154. return (_TENSOR_MODEL_PARALLEL_GROUP is not None
  155. and _PIPELINE_MODEL_PARALLEL_GROUP is not None)
  156. def get_cpu_world_group():
  157. """Get the CPU world group."""
  158. assert _CPU_WORLD_GROUP is not None, ("CPU world group is not initialized")
  159. return _CPU_WORLD_GROUP
  160. def get_tensor_model_parallel_group():
  161. """Get the tensor model parallel group the caller rank belongs to."""
  162. assert _TENSOR_MODEL_PARALLEL_GROUP is not None, (
  163. "tenosr model parallel group is not initialized")
  164. return _TENSOR_MODEL_PARALLEL_GROUP
  165. def get_pipeline_model_parallel_group():
  166. """Get the pipeline model parallel group the caller rank belongs to."""
  167. assert _PIPELINE_MODEL_PARALLEL_GROUP is not None, (
  168. "pipeline model parallel group is not initialized")
  169. return _PIPELINE_MODEL_PARALLEL_GROUP
  170. def get_tensor_model_parallel_world_size():
  171. """Return world size for the tensor model parallel group."""
  172. return torch.distributed.get_world_size(
  173. group=get_tensor_model_parallel_group())
  174. def get_pipeline_model_parallel_world_size():
  175. """Return world size for the pipeline model parallel group."""
  176. return torch.distributed.get_world_size(
  177. group=get_pipeline_model_parallel_group())
  178. def get_tensor_model_parallel_rank():
  179. """Return my rank for the tensor model parallel group."""
  180. return torch.distributed.get_rank(group=get_tensor_model_parallel_group())
  181. def get_pipeline_model_parallel_rank():
  182. """Return my rank for the pipeline model parallel group."""
  183. return torch.distributed.get_rank(
  184. group=get_pipeline_model_parallel_group())
  185. def get_tensor_model_parallel_src_rank():
  186. """Calculate the global rank corresponding to the first local rank
  187. in the tensor model parallel group."""
  188. global_rank = torch.distributed.get_rank()
  189. local_world_size = get_tensor_model_parallel_world_size()
  190. return (global_rank // local_world_size) * local_world_size
  191. def get_pipeline_model_parallel_first_rank():
  192. """Return the global rank of the first process in the pipeline for the
  193. current tensor parallel group"""
  194. assert _PIPELINE_GLOBAL_RANKS is not None, (
  195. "Pipeline parallel group is not initialized")
  196. return _PIPELINE_GLOBAL_RANKS[0]
  197. def get_pipeline_model_parallel_last_rank():
  198. """Return the global rank of the last process in the pipeline for the
  199. current tensor parallel group"""
  200. assert _PIPELINE_GLOBAL_RANKS is not None, (
  201. "Pipeline parallel group is not initialized")
  202. last_rank_local = get_pipeline_model_parallel_world_size() - 1
  203. return _PIPELINE_GLOBAL_RANKS[last_rank_local]
  204. def get_pipeline_model_parallel_next_rank():
  205. """Return the global rank that follows the caller in the pipeline"""
  206. assert _PIPELINE_GLOBAL_RANKS is not None, (
  207. "Pipeline parallel group is not initialized")
  208. rank_in_pipeline = get_pipeline_model_parallel_rank()
  209. world_size = get_pipeline_model_parallel_world_size()
  210. return _PIPELINE_GLOBAL_RANKS[(rank_in_pipeline + 1) % world_size]
  211. def get_pipeline_model_parallel_prev_rank():
  212. """Return the global rank that precedes the caller in the pipeline"""
  213. assert _PIPELINE_GLOBAL_RANKS is not None, (
  214. "Pipeline parallel group is not initialized")
  215. rank_in_pipeline = get_pipeline_model_parallel_rank()
  216. world_size = get_pipeline_model_parallel_world_size()
  217. return _PIPELINE_GLOBAL_RANKS[(rank_in_pipeline - 1) % world_size]
  218. def destroy_model_parallel():
  219. """Set the groups to none and destroy them."""
  220. global _TENSOR_MODEL_PARALLEL_GROUP
  221. if _TENSOR_MODEL_PARALLEL_GROUP:
  222. torch.distributed.destroy_process_group(_TENSOR_MODEL_PARALLEL_GROUP)
  223. _TENSOR_MODEL_PARALLEL_GROUP = None
  224. global _PIPELINE_MODEL_PARALLEL_GROUP
  225. if _PIPELINE_MODEL_PARALLEL_GROUP:
  226. torch.distributed.destroy_process_group(_PIPELINE_MODEL_PARALLEL_GROUP)
  227. _PIPELINE_MODEL_PARALLEL_GROUP = None
  228. global _PIPELINE_GLOBAL_RANKS
  229. _PIPELINE_GLOBAL_RANKS = None
  230. from aphrodite.distributed.device_communicators import pynccl_utils
  231. # Destroy the pynccl states if any.
  232. pynccl_utils.destroy_process_group()
  233. # Whether to use pynccl for nccl all reduce.
  234. # We use pynccl for all reduce when using CUDA graph, because torch.distributed
  235. # is not well supported by CUDA graph.
  236. _ENABLE_PYNCCL_FOR_ALL_REDUCE = False
  237. @contextlib.contextmanager
  238. def with_pynccl_for_all_reduce():
  239. """use Pynccl instead of torch.distributed for all reduce"""
  240. from aphrodite.distributed.device_communicators import pynccl_utils
  241. tp_size = get_tensor_model_parallel_world_size()
  242. if tp_size == 1:
  243. # No-op.
  244. # NOTE: We don't initialize Pynccl when tp_size is 1.
  245. yield
  246. else:
  247. global _ENABLE_PYNCCL_FOR_ALL_REDUCE
  248. old = _ENABLE_PYNCCL_FOR_ALL_REDUCE
  249. _ENABLE_PYNCCL_FOR_ALL_REDUCE = True
  250. stream = torch.cuda.current_stream()
  251. with pynccl_utils.set_pynccl_stream(stream):
  252. yield
  253. _ENABLE_PYNCCL_FOR_ALL_REDUCE = old
  254. def is_pynccl_enabled_for_all_reduce():
  255. """check if Pynccl is enabled for all reduce"""
  256. global _ENABLE_PYNCCL_FOR_ALL_REDUCE
  257. return _ENABLE_PYNCCL_FOR_ALL_REDUCE