accelerator_partitioner.py 47 KB

123456789101112131415161718192021222324252627282930313233343536373839404142434445464748495051525354555657585960616263646566676869707172737475767778798081828384858687888990919293949596979899100101102103104105106107108109110111112113114115116117118119120121122123124125126127128129130131132133134135136137138139140141142143144145146147148149150151152153154155156157158159160161162163164165166167168169170171172173174175176177178179180181182183184185186187188189190191192193194195196197198199200201202203204205206207208209210211212213214215216217218219220221222223224225226227228229230231232233234235236237238239240241242243244245246247248249250251252253254255256257258259260261262263264265266267268269270271272273274275276277278279280281282283284285286287288289290291292293294295296297298299300301302303304305306307308309310311312313314315316317318319320321322323324325326327328329330331332333334335336337338339340341342343344345346347348349350351352353354355356357358359360361362363364365366367368369370371372373374375376377378379380381382383384385386387388389390391392393394395396397398399400401402403404405406407408409410411412413414415416417418419420421422423424425426427428429430431432433434435436437438439440441442443444445446447448449450451452453454455456457458459460461462463464465466467468469470471472473474475476477478479480481482483484485486487488489490491492493494495496497498499500501502503504505506507508509510511512513514515516517518519520521522523524525526527528529530531532533534535536537538539540541542543544545546547548549550551552553554555556557558559560561562563564565566567568569570571572573574575576577578579580581582583584585586587588589590591592593594595596597598599600601602603604605606607608609610611612613614615616617618619620621622623624625626627628629630631632633634635636637638639640641642643644645646647648649650651652653654655656657658659660661662663664665666667668669670671672673674675676677678679680681682683684685686687688689690691692693694695696697698699700701702703704705706707708709710711712713714715716717718719720721722723724725726727728729730731732733734735736737738739740741742743744745746747748749750751752753754755756757758759760761762763764765766767768769770771772773774775776777778779780781782783784785786787788789790791792793794795796797798799800801802803804805806807808809810811812813814815816817818819820821822823824825826827828829830831832833834835836837838839840841842843844845846847848849850851852853854855856857858859860861862863864865866867868869870871872873874875876877878879880881882883884885886887888889890891892893894895896897898899900901902903904905906907908909910911912913914915916917918919920921922923924925926927928929930931932933934935936937938939940941942943944945946947948949950951952953954955956957958959960961962963964965966967968969970971972973974975976977978979980981982983984985986987988989990991992993994995996997998999100010011002100310041005100610071008100910101011101210131014101510161017101810191020102110221023102410251026102710281029103010311032103310341035103610371038103910401041104210431044104510461047104810491050105110521053105410551056105710581059106010611062106310641065106610671068106910701071107210731074107510761077107810791080108110821083
  1. import operator
  2. from collections import deque
  3. from typing import Dict, List, Set, NamedTuple, Tuple, Deque
  4. import torch
  5. from torch.fx.passes.graph_manipulation import get_size_of_all_nodes
  6. from torch.fx.experimental.partitioner_utils import (
  7. Partition,
  8. Device,
  9. PartitionerConfig,
  10. get_partition_to_latency_mapping,
  11. get_latency_of_partitioned_graph,
  12. NodeLatency,
  13. get_extra_size_of,
  14. PartitionMode,
  15. )
  16. from torch.fx.graph_module import GraphModule
  17. from torch.fx.node import Node, map_arg
  18. from torch.fx.passes.split_module import split_module
  19. class DAGNode:
  20. """DAGNode class maintains useful information for a partition (submodule),
  21. and its input submodules and output submodules.
  22. """
  23. def __init__(
  24. self,
  25. submodule_node: Node,
  26. input_nodes: List[Node],
  27. output_nodes: List[Node],
  28. logical_device_ids: List[int],
  29. size_bytes: int,
  30. ) -> None:
  31. self.submodule_node: Node = submodule_node
  32. self.input_nodes: List[Node] = input_nodes
  33. self.output_nodes: List[Node] = output_nodes
  34. self.logical_device_ids: List[int] = logical_device_ids
  35. self.size_bytes = size_bytes
  36. def __str__(self) -> str:
  37. return str(self.submodule_node)
  38. class DAG:
  39. """DAG class contains all the DAG nodes"""
  40. def __init__(self) -> None:
  41. self.nodes: List[DAGNode] = []
  42. def create_node(
  43. self,
  44. submodule_node: Node,
  45. input_nodes: List[Node],
  46. output_nodes: List[Node],
  47. logical_devices: List[int],
  48. size_bytes: int,
  49. ) -> None:
  50. node = DAGNode(
  51. submodule_node, input_nodes, output_nodes, logical_devices, size_bytes
  52. )
  53. self.nodes.append(node)
  54. class PartitionResult(NamedTuple):
  55. """NameTuple used for returning DAG and a new fx module"""
  56. dag: DAG
  57. module_with_submodules: GraphModule
  58. """Followings are some helper functions for partition manipulation"""
  59. def reset_partition_device(partitions):
  60. for partition in partitions:
  61. partition.logical_device_ids = []
  62. def combine_two_partitions(
  63. partition_0: Partition, partition_1: Partition, partitions: List[Partition]
  64. ) -> None:
  65. """Given a list of partitions and its two partitions,
  66. combine these two partitions into a new one appending to the partitions
  67. and remove the previous two partitions from the list of partitions
  68. """
  69. partition = Partition(len(partitions))
  70. partition.nodes = partition_0.nodes.union(partition_1.nodes)
  71. partition.recalculate_mem_size()
  72. partitions.append(partition)
  73. partitions.remove(partition_0)
  74. partitions.remove(partition_1)
  75. reorganize_partitions(partitions)
  76. return
  77. def set_parents_and_children(partitions: List[Partition]) -> None:
  78. """Given a list of partitions, mark parents and children for each partition"""
  79. # Go through all nodes in a partition.
  80. # If a node's user is in other partition,
  81. # then the other partition is this partition's children.
  82. # This partition is the other partition's parent
  83. for partition in partitions:
  84. partition.children = set()
  85. partition.parents = set()
  86. for partition in partitions:
  87. for node in partition.nodes:
  88. # For each node in the current partition, find its users
  89. users = node.users
  90. for n in users:
  91. # Find which the partition the user node belongs to.
  92. # Note that if the node itself is also belongs to that partition,
  93. # that partition is not the child of the current partition
  94. for p in partitions:
  95. if p != partition and n in p.nodes and node not in p.nodes:
  96. partition.children.add(p)
  97. p.parents.add(partition)
  98. return
  99. def reorganize_partitions(partitions: List[Partition]) -> None:
  100. """Given a list of partitions, reorganize partition id,
  101. its parents and its children for each partition
  102. """
  103. # Rearrange partition ids
  104. for i, partition in enumerate(partitions):
  105. partition.partition_id = i
  106. set_parents_and_children(partitions)
  107. return
  108. def get_bfs_level_partition(partitions: List[Partition]) -> None:
  109. """Given a list of partitions,
  110. mark the bfs level for each partition
  111. """
  112. current_level: Set[Partition] = set()
  113. visited: Set[Partition] = set()
  114. for partition in partitions:
  115. # If a partition has no parent, it should be in root level
  116. if len(partition.parents) == 0:
  117. current_level.add(partition)
  118. next_level: Set[Partition] = set()
  119. level = 0
  120. # bfs
  121. while current_level:
  122. partition = current_level.pop()
  123. partition.bfs_level = level
  124. visited.add(partition)
  125. children = partition.children
  126. for child in children:
  127. if child not in next_level:
  128. next_level.add(child)
  129. if not current_level:
  130. current_level = next_level.copy()
  131. next_level = set()
  132. level += 1
  133. return
  134. def get_node_to_partition_mapping(partitions: List[Partition]) -> Dict[Node, int]:
  135. """Given a list of partitions,return node to partition mapping"""
  136. node_to_partition: Dict[Node, int] = {}
  137. for partition in partitions:
  138. for node in partition.nodes:
  139. node_to_partition[node] = partition.partition_id
  140. return node_to_partition
  141. def get_logical_id_to_device(devices: List[Device]) -> Dict[int, Device]:
  142. """Get a mapping from device logical ID to Device object."""
  143. logical_id_to_device: Dict[int, Device] = {}
  144. for d in devices:
  145. logical_id_to_device[d.logical_id] = d
  146. return logical_id_to_device
  147. def get_device_partition_stats(
  148. partitions: List[Partition], devices: List[Device]
  149. ) -> Tuple[Dict[Device, List[Partition]], Dict[Device, int], List[Partition]]:
  150. """Given a list of partitions and a list of devices, returns:
  151. 1. A mapping from device to partitions on it;
  152. 2. A mapping from device to its remaining memory size;
  153. 3. A list of partitions that do not have a device.
  154. """
  155. # logical id to device
  156. logical_id_to_device = get_logical_id_to_device(devices)
  157. # Track partitions on device
  158. device_to_partitions: Dict[Device, List[Partition]] = {}
  159. # Track device's left mem size
  160. device_to_left_mem_bytes: Dict[Device, int] = {}
  161. for d in devices:
  162. device_to_partitions[d] = []
  163. device_to_left_mem_bytes[d] = d.available_mem_bytes
  164. # Deal with the partitions that already have a device
  165. # and also collect all partitions without a device (no_device_partitions)
  166. no_device_partitions = []
  167. for partition in partitions:
  168. if partition.logical_device_ids != []:
  169. for logical_id in partition.logical_device_ids:
  170. device = logical_id_to_device[logical_id]
  171. device_to_partitions[device].append(partition)
  172. device_to_left_mem_bytes[device] -= partition.used_mem_bytes
  173. else:
  174. no_device_partitions.append(partition)
  175. return (
  176. device_to_partitions,
  177. device_to_left_mem_bytes,
  178. no_device_partitions,
  179. )
  180. def get_device_to_partitions_mapping(
  181. partitions: List[Partition], devices: List[Device]
  182. ):
  183. """Given a list of partitions and a list of devices,
  184. map each partition into a device.
  185. """
  186. def calculate_extra_mem_bytes_needed_for(
  187. partition: Partition, partitions: List[Partition]
  188. ):
  189. all_nodes: Set[Node] = set()
  190. for p in partitions:
  191. all_nodes = all_nodes.union(p.nodes)
  192. if len(all_nodes) == 0:
  193. return partition.used_mem_bytes
  194. all_nodes = all_nodes.union(partition.nodes)
  195. extra_size_needed = 0
  196. for node in partition.nodes:
  197. extra_size_needed += get_extra_size_of(node, all_nodes)
  198. return extra_size_needed
  199. def find_device_for(partition: Partition):
  200. """Given a partition, find a logical device for the partition
  201. The algorithm is to put the partition on the device
  202. that has just enough mem left for that partition.
  203. device_to_left_mem_bytes is a dictionary between device and its left mem size
  204. sorted by its left mem size
  205. """
  206. for d in device_to_left_mem_bytes:
  207. extra_size_needed = calculate_extra_mem_bytes_needed_for(
  208. partition, device_to_partitions[d]
  209. )
  210. if extra_size_needed < device_to_left_mem_bytes[d]:
  211. device_to_partitions[d].append(partition)
  212. partition.logical_device_ids.append(d.logical_id)
  213. device_to_left_mem_bytes[d] -= extra_size_needed
  214. return True
  215. return False
  216. (
  217. device_to_partitions,
  218. device_to_left_mem_bytes,
  219. no_device_partitions,
  220. ) = get_device_partition_stats(partitions, devices)
  221. # Find devices for all the partitions without a device
  222. found_device = True
  223. for partition in no_device_partitions:
  224. device_to_left_mem_bytes = {
  225. d: left_mem_bytes
  226. for d, left_mem_bytes in sorted(
  227. device_to_left_mem_bytes.items(), key=lambda item: item[1]
  228. )
  229. }
  230. found_device = find_device_for(partition)
  231. if not found_device:
  232. break
  233. return found_device
  234. def check_dependency(partition):
  235. """Given a partition,check if there is a circular dependency on
  236. this partition using bfs
  237. """
  238. visited: Set[Partition] = {partition}
  239. queue: Deque[Partition] = deque([partition])
  240. while queue:
  241. p = queue.popleft()
  242. for child in p.children:
  243. if child == partition:
  244. return True
  245. else:
  246. if child not in visited:
  247. visited.add(child)
  248. queue.append(child)
  249. return False
  250. class Partitioner:
  251. """A fx module may not fit into one device.
  252. Partitioner class helps partition one fx module into submodules (partitions),
  253. so that the submodules can be executed crossing different accelerators.
  254. The main function of this class is self.partition_graph.
  255. It partitions the fx module based on the scheme specified in partition_config
  256. A DAG structure is returned
  257. along with a new fx module with submodule nodes.
  258. """
  259. def __init__(self) -> None:
  260. self.partitions: List[Partition] = []
  261. self.node_to_partition: Dict[Node, int] = {}
  262. self.devices: List[Device] = []
  263. def partition_graph(
  264. self,
  265. fx_module: GraphModule,
  266. torch_module: torch.nn.Module,
  267. partitioner_config: PartitionerConfig,
  268. ) -> PartitionResult:
  269. """Given the fx module, torch module and partitioner_config,
  270. find the partitions, do the partitions,
  271. and then return a DAG and a new fx module with submodule nodes (partitions)
  272. """
  273. self.graph_module = fx_module
  274. self.torch_module = torch_module
  275. self.devices = partitioner_config.devices
  276. if len(self.devices) == 0:
  277. raise RuntimeError("No devices")
  278. # Tag the size in bytes to all nodes in the graph_module.
  279. get_size_of_all_nodes(self.graph_module)
  280. # Check if there are op nodes in the fx module
  281. nodes = self.graph_module.graph.nodes
  282. if all(node.op in {"placeholder", "get_attr", "output"} for node in nodes):
  283. raise RuntimeError("No Partition since no operations in the module")
  284. # Calculate total size of the fx module
  285. total_size_of_graph = 0
  286. for node in nodes:
  287. if node.op == "output":
  288. break
  289. total_size_of_graph += node.size_bytes.total_size
  290. # Find the device with the max mem size
  291. device_with_max_mem = max(self.devices, key=lambda d: d.available_mem_bytes)
  292. # AOT based partition
  293. if partitioner_config.mode == PartitionMode.aot_based:
  294. self.aot_based_partition(
  295. partitioner_config.node_to_partition_mapping,
  296. partitioner_config.partition_to_logical_device_mapping,
  297. )
  298. # Single partition if the whole module can be fit into one device
  299. elif total_size_of_graph <= device_with_max_mem.available_mem_bytes:
  300. self.find_single_partition(
  301. total_size_of_graph, logical_device_id=device_with_max_mem.logical_id
  302. )
  303. elif total_size_of_graph > sum([d.available_mem_bytes for d in self.devices]):
  304. raise RuntimeError("Devices have no enough memory for the module")
  305. else:
  306. # Sparse nn based partition
  307. if partitioner_config.mode == PartitionMode.sparse_nn:
  308. available_mem_bytes = self.devices[0].available_mem_bytes
  309. if not all(
  310. device.available_mem_bytes == available_mem_bytes
  311. for device in self.devices
  312. ):
  313. raise RuntimeError("All devices must have same memory size!")
  314. # sparse_nn_partition only support same memory size
  315. # TODO: add different size support for sparse_nn_partition
  316. self.sparse_nn_partition(available_mem_bytes)
  317. # Cost aware partition
  318. elif partitioner_config.mode == PartitionMode.cost_aware:
  319. self.cost_aware_partition(
  320. partitioner_config.transfer_rate_bytes_per_sec,
  321. partitioner_config.node_to_latency_mapping,
  322. )
  323. # KL based partition
  324. elif partitioner_config.mode == PartitionMode.kl_based:
  325. self.kl_based_partition(
  326. partitioner_config.transfer_rate_bytes_per_sec,
  327. partitioner_config.node_to_latency_mapping,
  328. )
  329. else:
  330. self.size_based_partition()
  331. # Saturate host if possible.
  332. if partitioner_config.saturate_host:
  333. self.saturate_host()
  334. # Partition the graph module based on the partition assignment.
  335. module_with_submodules = self.do_partition()
  336. # The DAG contains DAGNodes with info of each partition's input nodes, output nodes
  337. # and how partitions are connected.
  338. dag = self.dump_dag(module_with_submodules)
  339. ret = PartitionResult(dag, module_with_submodules)
  340. return ret
  341. def find_single_partition(
  342. self, total_size_of_graph, logical_device_id: int = 0
  343. ) -> None:
  344. """Fit the whole fx module into one device"""
  345. partition_0 = self.create_partition()
  346. for node in self.graph_module.graph.nodes:
  347. if node.op == "output":
  348. # Skip the output node, but there can
  349. # be nodes after the output in certain cases.
  350. continue
  351. partition_0.nodes.add(node)
  352. partition_0.used_mem_bytes = total_size_of_graph
  353. partition_0.logical_device_ids = [logical_device_id]
  354. # Get the node to partition mapping
  355. self.node_to_partition = get_node_to_partition_mapping(self.partitions)
  356. return
  357. def size_based_partition(self) -> None:
  358. """This method is to partition the fx module based on memory size.
  359. It uses greedy approach. The result may not be the best.
  360. The basic idea is:
  361. Step 1:
  362. Find a device which has enough memory to fit the current node, create a empty partition
  363. with the size of that device.
  364. Then keep adding the following nodes into the partition until the partition is full.
  365. Step 2:
  366. Repeat Step 1 until no device left
  367. Step 3:
  368. If some nodes are left, create a partition for each left node (single node partition).
  369. and then try to map those partitions into logical devices with enough mem left.
  370. """
  371. def find_device_based_on_size(node) -> Device:
  372. """Given a node, this function is to find a logical device
  373. that could fit the node.
  374. """
  375. mem_size_needed = get_extra_size_of(node, set())
  376. device = Device("", -1, -1)
  377. for d in self.devices:
  378. if (
  379. d not in occupied_devices
  380. and d.available_mem_bytes >= mem_size_needed
  381. ):
  382. device = d
  383. break
  384. if device.available_mem_bytes < 0:
  385. raise RuntimeError(str(node) + "is too large to fit any device")
  386. occupied_devices.append(device)
  387. return device
  388. # Track partition and its left mem size
  389. partition_to_left_mem_bytes: Dict[Partition, int] = {}
  390. # Track all the devices that have been used
  391. occupied_devices: List[Device] = []
  392. partition = self.create_partition()
  393. for node in self.graph_module.graph.nodes:
  394. if node.op in {"call_module", "call_method", "call_function"}:
  395. # Check if there are devices left
  396. if len(self.partitions) <= len(self.devices):
  397. total_size_of_input_nodes = get_extra_size_of(node, partition.nodes)
  398. # Check if the current partition is the very first partition
  399. if partition.used_mem_bytes == 0:
  400. # Find a device to fit the first node, return available mem size
  401. device = find_device_based_on_size(node)
  402. occupied_devices.append(device)
  403. # Update partition and its left mem size
  404. partition_to_left_mem_bytes[
  405. partition
  406. ] = device.available_mem_bytes
  407. # Update available mem for the current partition
  408. partition.logical_device_ids.append(device.logical_id)
  409. else:
  410. # The current partition is not the first partition
  411. # Check if the current node can fit into current partition
  412. if (
  413. partition_to_left_mem_bytes[partition]
  414. < total_size_of_input_nodes
  415. ):
  416. # Check if no device is left
  417. if len(self.partitions) == len(self.devices):
  418. # No device is left
  419. # Put the previous partitions into a list (non_single_node_partitions)
  420. non_single_node_partitions = self.partitions[:]
  421. # Create the first single node partition for the current node
  422. self.create_single_node_partition(node)
  423. continue
  424. # Some devices are still left
  425. # Create a new partition with a mem size that is enough for the current node
  426. device = find_device_based_on_size(node)
  427. partition = self.create_partition()
  428. total_size_of_input_nodes = get_extra_size_of(
  429. node, partition.nodes
  430. )
  431. partition_to_left_mem_bytes[
  432. partition
  433. ] = device.available_mem_bytes
  434. partition.logical_device_ids.append(device.logical_id)
  435. partition.add_node(node)
  436. partition_to_left_mem_bytes[partition] -= total_size_of_input_nodes
  437. # Create single node partitions if no device is left
  438. else:
  439. self.create_single_node_partition(node)
  440. reorganize_partitions(self.partitions)
  441. # Get the node to partition mapping
  442. self.node_to_partition = get_node_to_partition_mapping(self.partitions)
  443. # Mapping all partitions into device
  444. found_partition_to_device_mapping = get_device_to_partitions_mapping(
  445. self.partitions, self.devices
  446. )
  447. if not found_partition_to_device_mapping:
  448. raise RuntimeError("Cannot Get a Valid Partition to Logical Device Mapping")
  449. return
  450. def saturate_host(self) -> None:
  451. """Saturate host by assigning replicates to unused devices with enough memory.
  452. It uses a greedy approach to find a next available set of devices to place all split
  453. partitions: For each used device, it searches for an idle device with minimal memory
  454. size that can hold all the partition located on that device; If the search is successful
  455. for all used devices, it then assigns the new devices' logical ID to the corresponding
  456. partition.
  457. """
  458. (
  459. device_to_partitions,
  460. device_to_left_mem_bytes,
  461. no_device_partitions,
  462. ) = get_device_partition_stats(self.partitions, self.devices)
  463. assert (
  464. len(no_device_partitions) == 0
  465. ), f"Expect no_device_partitions has 0 device, but get {len(no_device_partitions)}"
  466. # Devices that hold partitions
  467. used_devices = [d for d in self.devices if len(device_to_partitions[d]) > 0]
  468. # Track replicates of the assigned devices
  469. replicated_device_to_used_device: Dict[Device, Device] = {}
  470. while len(used_devices) * 2 + len(replicated_device_to_used_device) <= len(
  471. self.devices
  472. ):
  473. # Success flag for this round
  474. success = True
  475. # Devices that have not been assigned
  476. idle_devices = [
  477. d
  478. for d in self.devices
  479. if d not in used_devices and d not in replicated_device_to_used_device
  480. ]
  481. # Temporary mapping from replicated device to original device
  482. temp_replicate_mapping = {}
  483. # Find a new device to replicate all partitions on an used device
  484. for used_device in used_devices:
  485. # Idle devices that have enough memory
  486. available_devices = [
  487. d
  488. for d in idle_devices
  489. if d.available_mem_bytes
  490. >= used_device.available_mem_bytes
  491. - device_to_left_mem_bytes[used_device]
  492. ]
  493. if len(available_devices) == 0:
  494. success = False
  495. break
  496. new_device = min(available_devices, key=lambda d: d.available_mem_bytes)
  497. idle_devices.remove(new_device)
  498. temp_replicate_mapping[new_device] = used_device
  499. if not success:
  500. break
  501. replicated_device_to_used_device.update(temp_replicate_mapping)
  502. # Update logical device IDs assigned to the partitions
  503. for (
  504. replicate_device,
  505. original_device,
  506. ) in replicated_device_to_used_device.items():
  507. logical_id = replicate_device.logical_id
  508. for partition in device_to_partitions[original_device]:
  509. partition.logical_device_ids.append(logical_id)
  510. for p in self.partitions:
  511. print(p.logical_device_ids)
  512. def do_partition(self) -> GraphModule:
  513. """Return a new fx module with submodule nodes (partitions)."""
  514. module_with_submodules = split_module(
  515. self.graph_module,
  516. self.torch_module,
  517. lambda node: self.node_to_partition[node],
  518. )
  519. return module_with_submodules
  520. def dump_dag(self, module_with_submodules: GraphModule) -> DAG:
  521. """Return the dag structure and the new fx module with submodules."""
  522. dag = DAG()
  523. for node in module_with_submodules.graph.nodes:
  524. if node.op == "output":
  525. break
  526. if node.op in {"placeholder", "get_attr"}:
  527. continue
  528. if node.target == operator.__getitem__:
  529. continue
  530. input_nodes: Dict[Node, None] = {}
  531. map_arg(node.args, lambda n: input_nodes.setdefault(n))
  532. map_arg(node.kwargs, lambda n: input_nodes.setdefault(n))
  533. # When a node has two or more output nodes,
  534. # it outputs its result to 'getitem' nodes.
  535. # Those 'getitem' nodes are the output node for this node.
  536. # Otherwise, the output node is this node itself.
  537. if len(node.users) > 1:
  538. output_nodes = list(node.users)
  539. else:
  540. output_nodes = [node]
  541. partition_id = int(node.name.rsplit("_", 1)[-1])
  542. device_ids = self.partitions[partition_id].logical_device_ids
  543. size_bytes = self.partitions[partition_id].used_mem_bytes
  544. dag.create_node(
  545. node, list(input_nodes), output_nodes, device_ids, size_bytes
  546. )
  547. return dag
  548. def create_partition(self) -> Partition:
  549. """Create a partition and append it to self.partitions."""
  550. partition_id = len(self.partitions)
  551. partition = Partition(partition_id)
  552. self.partitions.append(partition)
  553. return partition
  554. def create_single_node_partition(self, node):
  555. """Create a partition for a single node"""
  556. partition = self.create_partition()
  557. partition.add_node(node)
  558. return
  559. def sparse_nn_partition(self, available_mem_bytes: int) -> None:
  560. """This method partition a sparse nn module.
  561. It is size based partition but different from size_based_partition,
  562. it only works when all the devices have same memory size (available_mem_bytes).
  563. In the future, devices with different mem sizes will be supported like size_based_partition.
  564. It first traverse all the nodes and do the partitions based on the same memory size.
  565. If the current partition has no enough memory left for a new op node
  566. (call_module, call_method, call_function), a new partition is created.
  567. When crossing the boundary between non-embedding nodes and embedding nodes,
  568. a new partition is created regardlessly.
  569. For example, if the current node is a non-embedding node but the next node is an
  570. embedding node, a new partition is created for the next node.
  571. After the partition, the partitions are combined as much as possible.
  572. The rule is that a non-embedding partition only
  573. combines with another non-embedding one.
  574. So as the embedding partitions.
  575. """
  576. def combine_partitions_based_on_size(
  577. partitions: List[Partition], available_mem_bytes: int
  578. ) -> None:
  579. """Combining small partitions together to keep as less partitions as possible.
  580. Here is an example of the algorithm to do this:
  581. Assume some partitions, we first sort them based on partition used memory size.
  582. [(partition_4, 1), (partition_3, 1), (partition_2, 2), (partition_1, 7), (partition_0, 9)]
  583. The available memory is 10.
  584. step 1: self.find_partition_to_combine_based_on_size()
  585. First, mark bfs level for each partition
  586. Second, look the smallest partition, partition_4: 10 - 1 = 9
  587. It means any partition has a used memory equal or less than 9 could combine this partition
  588. We go from the largest and selection partition_0.
  589. Check the bfs level for two partitions, if the level difference is less than 2,
  590. it can be combined.
  591. step 2: repeat step 1 until no partitions can be combined
  592. """
  593. find_combination = True
  594. while find_combination:
  595. # Sort partitions based on memory size
  596. sorted_partitions = sorted(partitions, key=lambda p: p.used_mem_bytes)
  597. # Mark bfs level
  598. get_bfs_level_partition(self.partitions)
  599. find_combination, partitions = find_partition_to_combine_based_on_size(
  600. sorted_partitions, available_mem_bytes, partitions
  601. )
  602. return
  603. def calculate_mem_bytes_needed(p1, p2):
  604. """Given two partitions, calculate how many mem bytes
  605. are needed if two partitions are combined
  606. """
  607. nodes = p1.nodes.union(p2.nodes)
  608. mem_bytes_needed = 0
  609. for node in nodes:
  610. mem_bytes_needed += get_extra_size_of(node, nodes)
  611. return mem_bytes_needed
  612. def find_partition_to_combine_based_on_size(
  613. sorted_partitions: List[Partition],
  614. available_mem_bytes: int,
  615. partitions: List[Partition],
  616. ) -> Tuple[bool, List[Partition]]:
  617. """step 1 in combine_partition_based_on_size()"""
  618. find_combination = False
  619. smallest_partition = sorted_partitions.pop(0)
  620. for p in sorted_partitions[::-1]:
  621. if abs(smallest_partition.bfs_level - p.bfs_level) <= 1:
  622. # Calculate how many bytes needed if combined
  623. mem_bytes_needed = calculate_mem_bytes_needed(p, smallest_partition)
  624. if mem_bytes_needed <= available_mem_bytes:
  625. combine_two_partitions(p, smallest_partition, self.partitions)
  626. partitions.remove(smallest_partition)
  627. partitions.remove(p)
  628. partitions.append(self.partitions[-1])
  629. find_combination = True
  630. break
  631. return find_combination, partitions
  632. def reset_partition_in_sparse_nn(partition, new_partition=True):
  633. """If crossing the boundary between non-embedding nodes and
  634. embedding nodes, create a new partition
  635. """
  636. if in_embedding_region:
  637. embedding_partitions.append(partition)
  638. else:
  639. non_embedding_partitions.append(partition)
  640. if new_partition:
  641. partition = self.create_partition()
  642. partition.left_mem_bytes = available_mem_bytes
  643. return partition
  644. return None
  645. def is_embedding_node(node: Node) -> bool:
  646. """Check if a node is an embedding node"""
  647. if node.op == "call_module":
  648. submodule = self.graph_module
  649. for atom in str(node.target).split("."):
  650. if not hasattr(submodule, atom):
  651. raise RuntimeError(
  652. f"Module {submodule} has no attribute {atom}"
  653. )
  654. submodule = getattr(submodule, atom)
  655. if "Embedding" in str(submodule):
  656. return True
  657. return False
  658. # Track embedding partitions and non-embedding partitions separately
  659. embedding_partitions: List[Partition] = []
  660. non_embedding_partitions: List[Partition] = []
  661. # A Flag to check the boundary
  662. in_embedding_region: bool = False
  663. partition = self.create_partition()
  664. for node in self.graph_module.graph.nodes:
  665. if node.op in {"call_module", "call_method", "call_function"}:
  666. # Check if crossing the boundary between embedding nodes and non embedding nodes
  667. if is_embedding_node(node) != in_embedding_region:
  668. # Crossing the boundary
  669. # Check if the current partition is an empty partition
  670. if partition.used_mem_bytes != 0:
  671. # The current partition isn't an empty partition. Create a new one.
  672. partition = reset_partition_in_sparse_nn(partition)
  673. in_embedding_region = not in_embedding_region
  674. total_size_of_input_nodes = get_extra_size_of(node, partition.nodes)
  675. if (
  676. total_size_of_input_nodes + partition.used_mem_bytes
  677. > available_mem_bytes
  678. ):
  679. partition = reset_partition_in_sparse_nn(partition)
  680. total_size_of_input_nodes = get_extra_size_of(node, partition.nodes)
  681. if total_size_of_input_nodes > available_mem_bytes:
  682. raise RuntimeError(
  683. node.target + "is too large to fit into a device"
  684. )
  685. partition.add_node(node)
  686. reset_partition_in_sparse_nn(partition, new_partition=False)
  687. # Set parents and children for partitions
  688. set_parents_and_children(self.partitions)
  689. # Combining non-embedding partitions
  690. combine_partitions_based_on_size(non_embedding_partitions, available_mem_bytes)
  691. # Combining embedding partitions
  692. combine_partitions_based_on_size(embedding_partitions, available_mem_bytes)
  693. total_size_of_non_embedding_partitions = 0
  694. for partition in non_embedding_partitions:
  695. total_size_of_non_embedding_partitions += partition.used_mem_bytes
  696. # Check if devices are enough for all partitions
  697. if len(embedding_partitions) > len(self.devices):
  698. msg = (
  699. "Need "
  700. + str(len(embedding_partitions))
  701. + " devices, but only "
  702. + str(len(self.devices))
  703. + " provided"
  704. )
  705. raise RuntimeError(msg)
  706. occupied_devices = []
  707. for i, partition in enumerate(embedding_partitions):
  708. # Check if all non-embedding partitions can fit into embedding partition devices
  709. if (
  710. total_size_of_non_embedding_partitions + partition.used_mem_bytes
  711. > available_mem_bytes
  712. ):
  713. raise RuntimeError(
  714. "partition_"
  715. + str(partition.partition_id)
  716. + "(embedding partition) and non embedding partitions can not fit into one device"
  717. )
  718. else:
  719. # Add logical device to the partition
  720. partition.logical_device_ids = [self.devices[i].logical_id]
  721. occupied_devices.append(self.devices[i].logical_id)
  722. # Add logical devices to the non_embedding_partitions
  723. for partition in non_embedding_partitions:
  724. partition.logical_device_ids = occupied_devices
  725. # Get the node to partition mapping
  726. self.node_to_partition = get_node_to_partition_mapping(self.partitions)
  727. return
  728. def cost_aware_partition(
  729. self,
  730. transfer_rate_bytes_per_sec: float,
  731. node_to_latency_mapping: Dict[Node, NodeLatency],
  732. ) -> None:
  733. """This method is to partition the fx module based on the cost.
  734. The cost is the total latency of running the whole fx module.
  735. In partitioner_utils.py, the cost model is built.
  736. The cost aware partition algorithm is:
  737. #1. At every beginning, each node is a partition.
  738. Then we map all the partitions to the devices
  739. and calculate the cost
  740. #2. Then try to pre-combine any two of the partitions if the two
  741. partitions can be combined.
  742. (the bfs level is less than 2 or two partitions are connected and
  743. can find partition to device mapping)
  744. See if any partition pair could reduce the current cost.
  745. Choose the pair that shows the minimum cost and then combine them
  746. #3. Repeat #2 until the cost cannot be reduced.
  747. """
  748. def try_combining_partitions(p0_index, p1_index, partitions) -> float:
  749. """Given two partitions and a list of partitions, combine these two partitions
  750. and see what is the cost of the modified partition list
  751. """
  752. p0 = partitions[p0_index]
  753. p1 = partitions[p1_index]
  754. """If two partitions' bfs level are less than 2 or two partitions are connected to each other,
  755. then they can be combined
  756. """
  757. if (
  758. (abs(p0.bfs_level - p1.bfs_level) <= 1)
  759. or (p0 in p1.parents)
  760. or p0 in (p1.children)
  761. ):
  762. combine_two_partitions(p0, p1, partitions)
  763. # Check if a circular dependency exists after combining
  764. if check_dependency(partitions[-1]):
  765. return float("inf")
  766. # Check if the modified partition list can be mapped to devices after combination
  767. reset_partition_device(partitions)
  768. found_deivce = get_device_to_partitions_mapping(
  769. partitions, self.devices
  770. )
  771. if not found_deivce:
  772. return float("inf")
  773. # Calculate the new cost
  774. partition_to_latency_mapping = get_partition_to_latency_mapping(
  775. partitions, node_to_latency_mapping
  776. )
  777. cost = get_latency_of_partitioned_graph(
  778. partitions,
  779. partition_to_latency_mapping,
  780. transfer_rate_bytes_per_sec,
  781. )
  782. return cost
  783. # If two partition can not be combined, the cost is inf
  784. return float("inf")
  785. def search_combination(
  786. transfer_rate_bytes_per_sec, node_to_latency_mapping
  787. ) -> bool:
  788. """Given transfer rate between partitions and each node's latency,
  789. find two partitions to combine so the cost of the partitions can
  790. be reduced.
  791. The algorithm is :
  792. 1. Go through all the partition pairs and see
  793. if any pair of partitions can be combined.
  794. 2. Calculate the cost after the combination.
  795. 3. Select the minimum cost and combine its corresponding partition pair.
  796. """
  797. partition_to_latency_mapping = get_partition_to_latency_mapping(
  798. self.partitions, node_to_latency_mapping
  799. )
  800. cost = get_latency_of_partitioned_graph(
  801. self.partitions,
  802. partition_to_latency_mapping,
  803. transfer_rate_bytes_per_sec,
  804. )
  805. if len(self.partitions) == 1:
  806. return False
  807. partition_pair: List[int] = []
  808. for i in range(len(self.partitions) - 1):
  809. for j in range(i + 1, len(self.partitions)):
  810. # Try to combine the partition pair
  811. # and see the new cost after combination
  812. new_cost = try_combining_partitions(i, j, self.partitions[:])
  813. if new_cost <= cost:
  814. partition_pair = [i, j]
  815. cost = new_cost
  816. reorganize_partitions(self.partitions)
  817. # If a partition pair is found, combine them
  818. if len(partition_pair) != 0:
  819. p0 = self.partitions[partition_pair[0]]
  820. p1 = self.partitions[partition_pair[1]]
  821. combine_two_partitions(p0, p1, self.partitions)
  822. get_bfs_level_partition(self.partitions)
  823. reset_partition_device(self.partitions)
  824. get_device_to_partitions_mapping(self.partitions, self.devices)
  825. return len(partition_pair) != 0
  826. for node in self.graph_module.graph.nodes:
  827. if node.op not in {"placeholder", "get_attr", "output"}:
  828. self.create_single_node_partition(node)
  829. # Set up parent partitions and children partitions for each partition
  830. set_parents_and_children(self.partitions)
  831. # Get bfs level for each partition
  832. get_bfs_level_partition(self.partitions)
  833. find_combination = True
  834. while find_combination:
  835. # Search for a pair partition to generate the minimum new cost,
  836. # then combine them
  837. find_combination = search_combination(
  838. transfer_rate_bytes_per_sec, node_to_latency_mapping
  839. )
  840. # Make sure all partitions are set up correctly
  841. reorganize_partitions(self.partitions)
  842. # Set up node to partition mapping
  843. self.node_to_partition = get_node_to_partition_mapping(self.partitions)
  844. return
  845. def kl_based_partition(
  846. self,
  847. transfer_rate_bytes_per_sec: float,
  848. node_to_latency_mapping: Dict[Node, NodeLatency],
  849. ) -> None:
  850. """This function is a cost aware partition based
  851. on Kernighan-Lin algorithm.
  852. First, the graph is partitioned using size_based_partition.
  853. Then, each node is swapped with any other node in a different
  854. partition, and at the same time, the cost is estimated after
  855. the swapping.
  856. For example, we have nodes n0, n1, n2, n3 and n4.
  857. Using size_based_partition, n0 and n1 are in Partition p0.
  858. n2, n3 and n4 in Partition p1. The current cost is estimated.
  859. We first tried using n0 to swap with n2 from the other partition.
  860. Then we see that swapping n0 and n2 shows a lower cost
  861. than the current cost and it is the minimum among other pairs like
  862. (n0, None)(This means moving n0 to Partition without swapping other nodes),
  863. (n0, n3) and (n0, n4). We swap n0 and n2 and set the new cost
  864. as the current cost.
  865. Then We repeat this process for all the other nodes until all swapping pairs
  866. are tried.
  867. """
  868. def swap_nodes(n0, n1, p0, p1):
  869. # Either n0 or n1 could be None
  870. # That means we simply move the node
  871. # to another partition
  872. if n0 is not None:
  873. p0.remove_node(n0)
  874. p1.add_node(n0)
  875. if n1 is not None:
  876. p0.add_node(n1)
  877. p1.remove_node(n1)
  878. def try_swap_nodes(
  879. n0, n1, p0, p1, node_to_latency_mapping, transfer_rate_per_sec
  880. ):
  881. cost = float("inf")
  882. swap_nodes(n0, n1, p0, p1)
  883. # Reorganize partitions after swapping
  884. reorganize_partitions(self.partitions)
  885. # Check if there is a circular dependency after swapping
  886. if (not check_dependency(p0)) and (not check_dependency(p1)):
  887. reset_partition_device(self.partitions)
  888. partition_to_latency_mapping = get_partition_to_latency_mapping(
  889. self.partitions, node_to_latency_mapping
  890. )
  891. # Check if all partitions can be mapped to logical devices after swapping
  892. found_device = get_device_to_partitions_mapping(
  893. self.partitions, self.devices
  894. )
  895. if not found_device:
  896. cost = float("inf")
  897. else:
  898. cost = get_latency_of_partitioned_graph(
  899. self.partitions,
  900. partition_to_latency_mapping,
  901. transfer_rate_bytes_per_sec,
  902. )
  903. # Swap back and reset all partitions back to original
  904. swap_nodes(n1, n0, p0, p1)
  905. reorganize_partitions(self.partitions)
  906. reset_partition_device(self.partitions)
  907. get_device_to_partitions_mapping(self.partitions, self.devices)
  908. return cost
  909. def swap_node_to_partition(
  910. node, p0, p1, node_to_latency_mapping, transfer_rate_per_sec
  911. ):
  912. """This function helps to swap one node from partition p0
  913. with all the nodes in another partition p1
  914. """
  915. p1_nodes = list(p1.nodes) + [None]
  916. min_cost = float("inf")
  917. node_pair: List[Node] = []
  918. for n1 in p1_nodes:
  919. # Ignore the node if it is not a op node
  920. if n1 is not None and n1.op in {"placeholder", "get_attr"}:
  921. continue
  922. # Try swapping node in p0 with n1 in p1
  923. cost = try_swap_nodes(
  924. node, n1, p0, p1, node_to_latency_mapping, transfer_rate_per_sec
  925. )
  926. if cost < min_cost:
  927. node_pair = [node, n1]
  928. min_cost = cost
  929. return cost, node_pair
  930. # First use size_base_partition
  931. self.size_based_partition()
  932. partition_to_latency_mapping = get_partition_to_latency_mapping(
  933. self.partitions, node_to_latency_mapping
  934. )
  935. # Calculate the cost of the partitions
  936. cost = get_latency_of_partitioned_graph(
  937. self.partitions, partition_to_latency_mapping, transfer_rate_bytes_per_sec
  938. )
  939. # Keep tracking the node pair that shows the better cost
  940. node_pair: List[Node] = []
  941. # Keep tracking the partition pair of node pair
  942. partition_pair: List[Partition] = []
  943. # Collect all the op nodes from the graph
  944. op_nodes = []
  945. for n in self.graph_module.graph.nodes:
  946. if n.op not in {"placeholder", "get_attr", "output"}:
  947. op_nodes.append(n)
  948. for node in op_nodes:
  949. # Find which partition the current node belongs
  950. p0_index = self.node_to_partition[node]
  951. p0 = self.partitions[p0_index]
  952. # Go through all the other partitions to swap
  953. # with other nodes from those partitions
  954. for p1_index, _ in enumerate(self.partitions):
  955. if p0_index != p1_index:
  956. p1 = self.partitions[p1_index]
  957. new_cost, new_node_pair = swap_node_to_partition(
  958. node,
  959. p0,
  960. p1,
  961. node_to_latency_mapping,
  962. transfer_rate_bytes_per_sec,
  963. )
  964. # Update the cost
  965. # Track the swapped node pair and their partitions
  966. if new_cost < cost:
  967. cost = new_cost
  968. node_pair = new_node_pair
  969. partition_pair = [p0, p1]
  970. # Do the swapping after trying all the nodes from a partition
  971. if len(node_pair) != 0:
  972. swap_nodes(
  973. node_pair[0], node_pair[1], partition_pair[0], partition_pair[1]
  974. )
  975. reorganize_partitions(self.partitions)
  976. get_device_to_partitions_mapping(self.partitions, self.devices)
  977. reorganize_partitions(self.partitions)
  978. # Mapping the device to the partition
  979. get_device_to_partitions_mapping(self.partitions, self.devices)
  980. return
  981. def aot_based_partition(
  982. self, node_to_partition_mapping, partition_to_logical_device_mapping
  983. ):
  984. """This function helps to rebuild the partitions given the nodes and its
  985. corresponding partition id
  986. """
  987. partition_id_to_partition_mapping: Dict[int, Partition] = {}
  988. self.node_to_partition = node_to_partition_mapping
  989. for node in self.node_to_partition:
  990. partition_id = self.node_to_partition[node]
  991. # If the requested partition has not been created, create the partition
  992. if partition_id not in partition_id_to_partition_mapping:
  993. partition = Partition(partition_id)
  994. self.partitions.append(partition)
  995. partition_id_to_partition_mapping[partition_id] = partition
  996. partition.logical_device_ids = partition_to_logical_device_mapping[
  997. partition_id
  998. ]
  999. else:
  1000. partition = partition_id_to_partition_mapping[
  1001. self.node_to_partition[node]
  1002. ]
  1003. # Add the current node into the partition
  1004. partition.add_node(node)