launch.py 14 KB

123456789101112131415161718192021222324252627282930313233343536373839404142434445464748495051525354555657585960616263646566676869707172737475767778798081828384858687888990919293949596979899100101102103104105106107108109110111112113114115116117118119120121122123124125126127128129130131132133134135136137138139140141142143144145146147148149150151152153154155156157158159160161162163164165166167168169170171172173174175176177178179180181182183184185186187188189190191192193194195196197198199200201202203204205206207208209210211212213214215216217218219220221222223224225226227228229230231232233234235236237238239240241242243244245246247248249250251252253254255256257258259260261262263264265266267268269270271272273274275276277278279280281282283284285286287288289290291292293294295296297298299300301302303304305306307308309310311312313314315316317318319320321322323324325326327328329330331332333334335336337338339340341342343344345346347348349350351352353354355356
  1. # Copyright (c) Microsoft Corporation.
  2. # SPDX-License-Identifier: Apache-2.0
  3. # DeepSpeed Team
  4. """
  5. DeepSpeed launcher, this is similar to torch's distributed.launch but supports
  6. additional features such as arbitrary gpu exclusion.
  7. deepspeed.launcher.launch is intended to be run on a single worker node and
  8. will spawn several worker sub-processes depending on how many devices/ranks
  9. are on the worker.
  10. """
  11. import sys
  12. import subprocess
  13. import os
  14. import json
  15. import base64
  16. import time
  17. import signal
  18. import psutil
  19. from collections import defaultdict
  20. from typing import Dict
  21. from argparse import ArgumentParser, REMAINDER
  22. from ..constants import TORCH_DISTRIBUTED_DEFAULT_PORT
  23. from ..nebula.constants import DLTS_POD_ENV_PATH
  24. from ..utils import logger, get_numactl_cmd
  25. from ..elasticity import is_torch_elastic_compatible
  26. from .constants import ELASTIC_TRAINING_ID_DEFAULT
  27. PID_FILE_BASEPATH = "/tmp"
  28. def parse_args():
  29. parser = ArgumentParser(description="DeepSpeed distributed training launch"
  30. " utility that creates multiple distributed"
  31. " processes on a single node")
  32. # Optional arguments for the launch helper
  33. parser.add_argument("--node_rank",
  34. type=int,
  35. default=0,
  36. help="The rank of the node for multi-node distributed "
  37. "training")
  38. parser.add_argument("--master_addr",
  39. default="127.0.0.1",
  40. type=str,
  41. help="Master node (rank 0)'s address, should be either"
  42. " the IP address or the hostname of node 0, for"
  43. " single node multi-proc training, the"
  44. " --master_addr can simply be 127.0.0.1")
  45. parser.add_argument("--master_port",
  46. default=TORCH_DISTRIBUTED_DEFAULT_PORT,
  47. type=int,
  48. help="Master node (rank 0)'s free port that needs to "
  49. "be used for communication during distributed "
  50. "training")
  51. parser.add_argument("--world_info", default="None", type=str, help="world info base64 encoded dictionary")
  52. parser.add_argument("--module",
  53. action="store_true",
  54. help="Change each process to interpret the launch "
  55. "script as a Python module, executing with the same "
  56. "behavior as 'python -m'.")
  57. parser.add_argument("--no_python",
  58. action="store_true",
  59. help="Skip prepending the training script with "
  60. "'python' - just execute it directly.")
  61. parser.add_argument("--enable_elastic_training", action="store_true", help="Enable elastic training support.")
  62. parser.add_argument("--min_elastic_nodes", type=int, default=-1, help="Min number of nodes in elastic training.")
  63. parser.add_argument("--max_elastic_nodes", type=int, default=-1, help="Max number of nodes in elastic training.")
  64. parser.add_argument("--no_local_rank",
  65. action="store_true",
  66. help="Do not pass local_rank as an argument when calling "
  67. "the user's training script.")
  68. parser.add_argument("--save_pid",
  69. type=int,
  70. default=0,
  71. help="main launching process pid, for internal pid tracking")
  72. parser.add_argument("--enable_each_rank_log",
  73. default="None",
  74. type=str,
  75. help="redirect the stdout and stderr from each rank into different log files")
  76. parser.add_argument("--bind_cores_to_rank",
  77. action="store_true",
  78. help="Bind each rank to different cores of the host. "
  79. "This improves host efficiency especially for CPU backend")
  80. parser.add_argument("--bind_core_list",
  81. type=str,
  82. default=None,
  83. help="List of cores to bind to with comma separated list of "
  84. "numbers and range. i.e. 1,3-5,7 => [1,3,4,5,7]. When not "
  85. "specified, all cores on system would be used rank binding")
  86. # positional
  87. parser.add_argument("training_script",
  88. type=str,
  89. help="The full path to the single GPU training "
  90. "program/script to be launched in parallel, "
  91. "followed by all the arguments for the "
  92. "training script")
  93. # rest from the training program
  94. parser.add_argument('training_script_args', nargs=REMAINDER)
  95. return parser.parse_args()
  96. # Adapted from https://psutil.readthedocs.io/en/latest/#kill-process-tree
  97. def terminate_process_tree(pid):
  98. process = psutil.Process(pid)
  99. children = process.children(recursive=True)
  100. children.append(process)
  101. for child in children:
  102. try:
  103. child.terminate()
  104. except psutil.NoSuchProcess:
  105. pass
  106. gone, alive = psutil.wait_procs(children, timeout=30)
  107. for p in alive:
  108. p.kill()
  109. def main():
  110. args = parse_args()
  111. current_env = os.environ.copy()
  112. for k in current_env.keys():
  113. if "NCCL" in k:
  114. logger.info(f"{args.node_rank} {k}={current_env[k]}")
  115. if args.world_info == "None":
  116. raise ValueError("world_info can not be None")
  117. world_info = base64.urlsafe_b64decode(args.world_info)
  118. world_info = json.loads(world_info)
  119. logger.info(f"WORLD INFO DICT: {world_info}")
  120. node_list = list(world_info.keys())
  121. args.nnodes = len(node_list)
  122. local_node = node_list[args.node_rank]
  123. local_gpu_ids = world_info[local_node]
  124. num_local_procs = len(local_gpu_ids)
  125. logger.info(f"nnodes={args.nnodes}, num_local_procs={num_local_procs}, node_rank={args.node_rank}")
  126. global_rank_mapping = defaultdict(list)
  127. curr_global_rank = 0
  128. dist_world_size = 0
  129. for node_id in node_list:
  130. gids = world_info[node_id]
  131. dist_world_size += len(gids)
  132. for gid in gids:
  133. global_rank_mapping[node_id].append(curr_global_rank)
  134. curr_global_rank += 1
  135. logger.info(f"global_rank_mapping={global_rank_mapping}")
  136. logger.info(f"dist_world_size={dist_world_size}")
  137. current_env["CUDA_VISIBLE_DEVICES"] = ",".join(map(str, local_gpu_ids))
  138. logger.info(f"Setting CUDA_VISIBLE_DEVICES={current_env['CUDA_VISIBLE_DEVICES']}")
  139. # set PyTorch distributed related environmental variables
  140. current_env["MASTER_ADDR"] = args.master_addr
  141. current_env["MASTER_PORT"] = str(args.master_port)
  142. current_env["WORLD_SIZE"] = str(dist_world_size)
  143. current_env["CROSS_RANK"] = str(args.node_rank)
  144. current_env["CROSS_SIZE"] = str(args.nnodes)
  145. current_env["LOCAL_SIZE"] = str(num_local_procs)
  146. if args.save_pid:
  147. print(f"launcher pid: {os.getpid()}")
  148. pid_file = None
  149. if args.save_pid:
  150. launcher_pid = os.getpid()
  151. pid_file = os.path.join(PID_FILE_BASEPATH, f"{args.save_pid}.deepspeed")
  152. assert not os.path.isfile(pid_file), "pid file exists but shouldn't"
  153. with open(pid_file, 'w') as fd:
  154. fd.write(f"{launcher_pid}")
  155. if not is_torch_elastic_compatible():
  156. if args.enable_elastic_training:
  157. logger.info(f"Disabling elastic training support as \
  158. PyTorch version should be greater than 1.11.x")
  159. args.enable_elastic_training = False
  160. if os.path.exists(DLTS_POD_ENV_PATH):
  161. with open(DLTS_POD_ENV_PATH) as file:
  162. lines = file.readlines()
  163. lines = [line.rstrip() for line in lines]
  164. for line in lines:
  165. if line.startswith('export FC_TASKROLE_NAME') or line.startswith('export FC_TASK_INDEX'):
  166. key_val = line.split()[1]
  167. key, val = key_val.split('=')
  168. current_env[key] = val
  169. processes = []
  170. cmd = []
  171. if not args.enable_elastic_training:
  172. if args.enable_each_rank_log != "None":
  173. # prepare the log path and the file name prefix
  174. if os.path.isfile(args.enable_each_rank_log):
  175. raise ValueError(f"{args.enable_each_rank_log} should not be a file, it should be a directory.")
  176. if not os.path.exists(args.enable_each_rank_log):
  177. try:
  178. os.makedirs(args.enable_each_rank_log)
  179. except Exception as e:
  180. print(e)
  181. raise ValueError(f"unable to create directory {args.enable_each_rank_log} for each rank log.")
  182. log_name_prefix = time.strftime("%Y%m%d%H%M%S", time.localtime())
  183. for local_proc in range(0, num_local_procs):
  184. # each process's rank
  185. dist_rank = global_rank_mapping[local_node][local_proc]
  186. local_rank = dist_rank % num_local_procs
  187. current_env["RANK"] = str(dist_rank)
  188. current_env["LOCAL_RANK"] = str(local_rank)
  189. # spawn the processes
  190. cmd = []
  191. if args.bind_cores_to_rank:
  192. cores_per_rank, numactl_cmd = get_numactl_cmd(args.bind_core_list, num_local_procs, local_rank)
  193. current_env["OMP_NUM_THREADS"] = f"{cores_per_rank}"
  194. cmd = cmd + numactl_cmd
  195. if not args.no_python:
  196. cmd.append(sys.executable)
  197. cmd.append("-u")
  198. if args.module:
  199. cmd.append("-m")
  200. else:
  201. if args.module:
  202. raise ValueError("Don't use both the '--no_python' flag"
  203. " and the '--module' flag at the same time.")
  204. cmd.append(args.training_script)
  205. # A user may not want to pass local_rank as a keyword arg so we make this optional.
  206. if not args.no_local_rank:
  207. cmd.append(f"--local_rank={local_rank}")
  208. cmd += args.training_script_args
  209. if args.enable_each_rank_log != "None":
  210. log_file = os.path.join(args.enable_each_rank_log, f"{log_name_prefix}_rank{dist_rank}.log")
  211. log_fd = open(log_file, 'w')
  212. process = subprocess.Popen(cmd, env=current_env, stdout=log_fd, stderr=log_fd)
  213. else:
  214. process = subprocess.Popen(cmd, env=current_env)
  215. # logs the command from processes
  216. logger.info(f"process {process.pid} spawned with command: {cmd}")
  217. processes.append(process)
  218. else:
  219. from ..elasticity import DSElasticAgent
  220. from torch.distributed.elastic.rendezvous import RendezvousParameters
  221. from torch.distributed.elastic.agent.server.api import WorkerSpec
  222. import torch.distributed.elastic.rendezvous.registry as rdzv_registry
  223. from torch.distributed.elastic.multiprocessing import Std
  224. if args.min_elastic_nodes == -1:
  225. args.min_elastic_nodes = 1
  226. if args.max_elastic_nodes == -1:
  227. args.max_elastic_nodes = args.nnodes
  228. assert args.max_elastic_nodes > 0 and args.min_elastic_nodes > 0, "Max and Min nodes should be positive"
  229. current_env["NCCL_ASYNC_ERROR_HANDLING"] = str(1)
  230. # Get config and arguments
  231. cmd = []
  232. if not args.no_python:
  233. cmd = [sys.executable, "-u"]
  234. if args.module:
  235. cmd.append("-m")
  236. else:
  237. if args.module:
  238. raise ValueError("Don't use both the '--no_python' flag"
  239. " and the '--module' flag at the same time.")
  240. cmd.append(args.training_script)
  241. cmd += args.training_script_args
  242. cmd_args = cmd[1:]
  243. rdzv_configs: Dict[str, str] = {'timeout': 100}
  244. run_id = os.environ.get("ELASTIC_RUN_ID", ELASTIC_TRAINING_ID_DEFAULT)
  245. # Creating config for rendezvous class
  246. rdzv_parameters = RendezvousParameters(backend='c10d',
  247. endpoint=args.master_addr + ":" + str(args.master_port),
  248. run_id=run_id,
  249. min_nodes=args.min_elastic_nodes,
  250. max_nodes=args.max_elastic_nodes,
  251. **rdzv_configs)
  252. spec = WorkerSpec(
  253. role='trainer',
  254. local_world_size=num_local_procs,
  255. entrypoint=cmd[0],
  256. args=cmd[1:],
  257. rdzv_handler=rdzv_registry.get_rendezvous_handler(rdzv_parameters),
  258. max_restarts=100,
  259. monitor_interval=5,
  260. redirects=Std.from_str("0"),
  261. tee=Std.from_str("0"),
  262. master_addr=None,
  263. master_port=None,
  264. )
  265. agent = DSElasticAgent(spec, current_env)
  266. agent.run()
  267. sig_names = {2: "SIGINT", 15: "SIGTERM"}
  268. last_return_code = None
  269. def sigkill_handler(signum, frame):
  270. for process in processes:
  271. logger.info(f"Killing subprocess {process.pid}")
  272. try:
  273. terminate_process_tree(process.pid)
  274. except Exception:
  275. pass
  276. if last_return_code is not None:
  277. logger.error(f"{cmd} exits with return code = {last_return_code}")
  278. sys.exit(last_return_code)
  279. if signum in sig_names:
  280. logger.info(f"Main process received {sig_names[signum]}, exiting")
  281. if args.save_pid:
  282. if os.path.isfile(pid_file):
  283. os.remove(pid_file)
  284. sys.exit(1)
  285. # pass SIGINT/SIGTERM to children if the parent is being terminated
  286. signal.signal(signal.SIGINT, sigkill_handler)
  287. signal.signal(signal.SIGTERM, sigkill_handler)
  288. alive_processes = set(processes)
  289. while len(alive_processes):
  290. finished_processes = []
  291. for process in alive_processes:
  292. if process.poll() is None:
  293. # the process is still running
  294. continue
  295. else:
  296. if process.returncode != 0:
  297. last_return_code = process.returncode # for sigkill_handler
  298. sigkill_handler(signal.SIGTERM, None) # not coming back
  299. else:
  300. # exited cleanly
  301. logger.info(f"Process {process.pid} exits successfully.")
  302. finished_processes.append(process)
  303. alive_processes = set(alive_processes) - set(finished_processes)
  304. time.sleep(1)
  305. if __name__ == "__main__":
  306. main()