builder.py 26 KB

123456789101112131415161718192021222324252627282930313233343536373839404142434445464748495051525354555657585960616263646566676869707172737475767778798081828384858687888990919293949596979899100101102103104105106107108109110111112113114115116117118119120121122123124125126127128129130131132133134135136137138139140141142143144145146147148149150151152153154155156157158159160161162163164165166167168169170171172173174175176177178179180181182183184185186187188189190191192193194195196197198199200201202203204205206207208209210211212213214215216217218219220221222223224225226227228229230231232233234235236237238239240241242243244245246247248249250251252253254255256257258259260261262263264265266267268269270271272273274275276277278279280281282283284285286287288289290291292293294295296297298299300301302303304305306307308309310311312313314315316317318319320321322323324325326327328329330331332333334335336337338339340341342343344345346347348349350351352353354355356357358359360361362363364365366367368369370371372373374375376377378379380381382383384385386387388389390391392393394395396397398399400401402403404405406407408409410411412413414415416417418419420421422423424425426427428429430431432433434435436437438439440441442443444445446447448449450451452453454455456457458459460461462463464465466467468469470471472473474475476477478479480481482483484485486487488489490491492493494495496497498499500501502503504505506507508509510511512513514515516517518519520521522523524525526527528529530531532533534535536537538539540541542543544545546547548549550551552553554555556557558559560561562563564565566567568569570571572573574575576577578579580581582583584585586587588589590591592593594595596597598599600601602603604605606607608609610611612613614615616617618619620621622623624625626627628629630631632633634635636637638639640641642643644645646647648649650651652653654655656657658659660661662663664665666667668669670671672673674675676677678679680681682683684685686687688689690691692693694695696697698699
  1. """
  2. Copyright 2020 The Microsoft DeepSpeed Team
  3. """
  4. import os
  5. import sys
  6. import time
  7. import importlib
  8. from pathlib import Path
  9. import subprocess
  10. import shlex
  11. import shutil
  12. import tempfile
  13. import distutils.ccompiler
  14. import distutils.log
  15. import distutils.sysconfig
  16. from distutils.errors import CompileError, LinkError
  17. from abc import ABC, abstractmethod
  18. from typing import List
  19. YELLOW = '\033[93m'
  20. END = '\033[0m'
  21. WARNING = f"{YELLOW} [WARNING] {END}"
  22. DEFAULT_TORCH_EXTENSION_PATH = "/tmp/torch_extensions"
  23. DEFAULT_COMPUTE_CAPABILITIES = "6.0;6.1;7.0"
  24. try:
  25. import torch
  26. except ImportError:
  27. print(
  28. f"{WARNING} unable to import torch, please install it if you want to pre-compile any deepspeed ops."
  29. )
  30. else:
  31. TORCH_MAJOR = int(torch.__version__.split('.')[0])
  32. TORCH_MINOR = int(torch.__version__.split('.')[1])
  33. def installed_cuda_version():
  34. import torch.utils.cpp_extension
  35. cuda_home = torch.utils.cpp_extension.CUDA_HOME
  36. assert cuda_home is not None, "CUDA_HOME does not exist, unable to compile CUDA op(s)"
  37. # Ensure there is not a cuda version mismatch between torch and nvcc compiler
  38. output = subprocess.check_output([cuda_home + "/bin/nvcc",
  39. "-V"],
  40. universal_newlines=True)
  41. output_split = output.split()
  42. release_idx = output_split.index("release")
  43. release = output_split[release_idx + 1].replace(',', '').split(".")
  44. # Ignore patch versions, only look at major + minor
  45. cuda_major, cuda_minor = release[:2]
  46. installed_cuda_version = ".".join(release[:2])
  47. return int(cuda_major), int(cuda_minor)
  48. def get_default_compute_capabilities():
  49. compute_caps = DEFAULT_COMPUTE_CAPABILITIES
  50. import torch.utils.cpp_extension
  51. if torch.utils.cpp_extension.CUDA_HOME is not None and installed_cuda_version(
  52. )[0] >= 11:
  53. if installed_cuda_version()[0] == 11 and installed_cuda_version()[1] == 0:
  54. # Special treatment of CUDA 11.0 because compute_86 is not supported.
  55. compute_caps += ";8.0"
  56. else:
  57. compute_caps += ";8.0;8.6"
  58. return compute_caps
  59. # list compatible minor CUDA versions - so that for example pytorch built with cuda-11.0 can be used
  60. # to build deepspeed and system-wide installed cuda 11.2
  61. cuda_minor_mismatch_ok = {
  62. 10: [
  63. "10.0",
  64. "10.1",
  65. "10.2",
  66. ],
  67. 11: ["11.0",
  68. "11.1",
  69. "11.2",
  70. "11.3",
  71. "11.4",
  72. "11.5",
  73. "11.6",
  74. "11.7",
  75. "11.8"],
  76. }
  77. def assert_no_cuda_mismatch():
  78. cuda_major, cuda_minor = installed_cuda_version()
  79. sys_cuda_version = f'{cuda_major}.{cuda_minor}'
  80. torch_cuda_version = ".".join(torch.version.cuda.split('.')[:2])
  81. # This is a show-stopping error, should probably not proceed past this
  82. if sys_cuda_version != torch_cuda_version:
  83. if (cuda_major in cuda_minor_mismatch_ok
  84. and sys_cuda_version in cuda_minor_mismatch_ok[cuda_major]
  85. and torch_cuda_version in cuda_minor_mismatch_ok[cuda_major]):
  86. print(f"Installed CUDA version {sys_cuda_version} does not match the "
  87. f"version torch was compiled with {torch.version.cuda} "
  88. "but since the APIs are compatible, accepting this combination")
  89. return
  90. raise Exception(
  91. f"Installed CUDA version {sys_cuda_version} does not match the "
  92. f"version torch was compiled with {torch.version.cuda}, unable to compile "
  93. "cuda/cpp extensions without a matching cuda version.")
  94. class OpBuilder(ABC):
  95. _rocm_version = None
  96. _is_rocm_pytorch = None
  97. def __init__(self, name):
  98. self.name = name
  99. self.jit_mode = False
  100. self.error_log = None
  101. @abstractmethod
  102. def absolute_name(self):
  103. '''
  104. Returns absolute build path for cases where the op is pre-installed, e.g., deepspeed.ops.adam.cpu_adam
  105. will be installed as something like: deepspeed/ops/adam/cpu_adam.so
  106. '''
  107. pass
  108. @abstractmethod
  109. def sources(self):
  110. '''
  111. Returns list of source files for your op, relative to root of deepspeed package (i.e., DeepSpeed/deepspeed)
  112. '''
  113. pass
  114. def hipify_extension(self):
  115. pass
  116. @staticmethod
  117. def assert_torch_info(torch_info):
  118. install_torch_version = torch_info['version']
  119. install_cuda_version = torch_info['cuda_version']
  120. install_hip_version = torch_info['hip_version']
  121. if not OpBuilder.is_rocm_pytorch():
  122. current_cuda_version = ".".join(torch.version.cuda.split('.')[:2])
  123. else:
  124. current_hip_version = ".".join(torch.version.hip.split('.')[:2])
  125. current_torch_version = ".".join(torch.__version__.split('.')[:2])
  126. if not OpBuilder.is_rocm_pytorch():
  127. if install_cuda_version != current_cuda_version or install_torch_version != current_torch_version:
  128. raise RuntimeError(
  129. "PyTorch and CUDA version mismatch! DeepSpeed ops were compiled and installed "
  130. "with a different version than what is being used at runtime. Please re-install "
  131. f"DeepSpeed or switch torch versions. DeepSpeed install versions: "
  132. f"torch={install_torch_version}, cuda={install_cuda_version}, runtime versions:"
  133. f"torch={current_torch_version}, cuda={current_cuda_version}")
  134. else:
  135. if install_hip_version != current_hip_version or install_torch_version != current_torch_version:
  136. raise RuntimeError(
  137. "PyTorch and HIP version mismatch! DeepSpeed ops were compiled and installed "
  138. "with a different version than what is being used at runtime. Please re-install "
  139. f"DeepSpeed or switch torch versions. DeepSpeed install versions: "
  140. f"torch={install_torch_version}, hip={install_hip_version}, runtime versions:"
  141. f"torch={current_torch_version}, hip={current_hip_version}")
  142. @staticmethod
  143. def is_rocm_pytorch():
  144. if OpBuilder._is_rocm_pytorch is not None:
  145. return OpBuilder._is_rocm_pytorch
  146. _is_rocm_pytorch = False
  147. try:
  148. import torch
  149. except ImportError:
  150. pass
  151. else:
  152. if TORCH_MAJOR > 1 or (TORCH_MAJOR == 1 and TORCH_MINOR >= 5):
  153. _is_rocm_pytorch = hasattr(torch.version,
  154. 'hip') and torch.version.hip is not None
  155. if _is_rocm_pytorch:
  156. from torch.utils.cpp_extension import ROCM_HOME
  157. _is_rocm_pytorch = ROCM_HOME is not None
  158. OpBuilder._is_rocm_pytorch = _is_rocm_pytorch
  159. return OpBuilder._is_rocm_pytorch
  160. @staticmethod
  161. def installed_rocm_version():
  162. if OpBuilder._rocm_version:
  163. return OpBuilder._rocm_version
  164. ROCM_MAJOR = '0'
  165. ROCM_MINOR = '0'
  166. if OpBuilder.is_rocm_pytorch():
  167. from torch.utils.cpp_extension import ROCM_HOME
  168. rocm_ver_file = Path(ROCM_HOME).joinpath(".info/version-dev")
  169. if rocm_ver_file.is_file():
  170. with open(rocm_ver_file, 'r') as file:
  171. ROCM_VERSION_DEV_RAW = file.read()
  172. elif "rocm" in torch.__version__:
  173. ROCM_VERSION_DEV_RAW = torch.__version__.split("rocm")[1]
  174. else:
  175. assert False, "Could not detect ROCm version"
  176. assert ROCM_VERSION_DEV_RAW != "", "Could not detect ROCm version"
  177. ROCM_MAJOR = ROCM_VERSION_DEV_RAW.split('.')[0]
  178. ROCM_MINOR = ROCM_VERSION_DEV_RAW.split('.')[1]
  179. OpBuilder._rocm_version = (int(ROCM_MAJOR), int(ROCM_MINOR))
  180. return OpBuilder._rocm_version
  181. def include_paths(self):
  182. '''
  183. Returns list of include paths, relative to root of deepspeed package (i.e., DeepSpeed/deepspeed)
  184. '''
  185. return []
  186. def nvcc_args(self):
  187. '''
  188. Returns optional list of compiler flags to forward to nvcc when building CUDA sources
  189. '''
  190. return []
  191. def cxx_args(self):
  192. '''
  193. Returns optional list of compiler flags to forward to the build
  194. '''
  195. return []
  196. def is_compatible(self, verbose=True):
  197. '''
  198. Check if all non-python dependencies are satisfied to build this op
  199. '''
  200. return True
  201. def extra_ldflags(self):
  202. return []
  203. def libraries_installed(self, libraries):
  204. valid = False
  205. check_cmd = 'dpkg -l'
  206. for lib in libraries:
  207. result = subprocess.Popen(f'dpkg -l {lib}',
  208. stdout=subprocess.PIPE,
  209. stderr=subprocess.PIPE,
  210. shell=True)
  211. valid = valid or result.wait() == 0
  212. return valid
  213. def has_function(self, funcname, libraries, verbose=False):
  214. '''
  215. Test for existence of a function within a tuple of libraries.
  216. This is used as a smoke test to check whether a certain library is available.
  217. As a test, this creates a simple C program that calls the specified function,
  218. and then distutils is used to compile that program and link it with the specified libraries.
  219. Returns True if both the compile and link are successful, False otherwise.
  220. '''
  221. tempdir = None # we create a temporary directory to hold various files
  222. filestderr = None # handle to open file to which we redirect stderr
  223. oldstderr = None # file descriptor for stderr
  224. try:
  225. # Echo compile and link commands that are used.
  226. if verbose:
  227. distutils.log.set_verbosity(1)
  228. # Create a compiler object.
  229. compiler = distutils.ccompiler.new_compiler(verbose=verbose)
  230. # Configure compiler and linker to build according to Python install.
  231. distutils.sysconfig.customize_compiler(compiler)
  232. # Create a temporary directory to hold test files.
  233. tempdir = tempfile.mkdtemp()
  234. # Define a simple C program that calls the function in question
  235. prog = "void %s(void); int main(int argc, char** argv) { %s(); return 0; }" % (
  236. funcname,
  237. funcname)
  238. # Write the test program to a file.
  239. filename = os.path.join(tempdir, 'test.c')
  240. with open(filename, 'w') as f:
  241. f.write(prog)
  242. # Redirect stderr file descriptor to a file to silence compile/link warnings.
  243. if not verbose:
  244. filestderr = open(os.path.join(tempdir, 'stderr.txt'), 'w')
  245. oldstderr = os.dup(sys.stderr.fileno())
  246. os.dup2(filestderr.fileno(), sys.stderr.fileno())
  247. # Workaround for behavior in distutils.ccompiler.CCompiler.object_filenames()
  248. # Otherwise, a local directory will be used instead of tempdir
  249. drive, driveless_filename = os.path.splitdrive(filename)
  250. root_dir = driveless_filename[0] if os.path.isabs(driveless_filename) else ''
  251. output_dir = os.path.join(drive, root_dir)
  252. # Attempt to compile the C program into an object file.
  253. cflags = shlex.split(os.environ.get('CFLAGS', ""))
  254. objs = compiler.compile([filename],
  255. output_dir=output_dir,
  256. extra_preargs=self.strip_empty_entries(cflags))
  257. # Attempt to link the object file into an executable.
  258. # Be sure to tack on any libraries that have been specified.
  259. ldflags = shlex.split(os.environ.get('LDFLAGS', ""))
  260. compiler.link_executable(objs,
  261. os.path.join(tempdir,
  262. 'a.out'),
  263. extra_preargs=self.strip_empty_entries(ldflags),
  264. libraries=libraries)
  265. # Compile and link succeeded
  266. return True
  267. except CompileError:
  268. return False
  269. except LinkError:
  270. return False
  271. except:
  272. return False
  273. finally:
  274. # Restore stderr file descriptor and close the stderr redirect file.
  275. if oldstderr is not None:
  276. os.dup2(oldstderr, sys.stderr.fileno())
  277. if filestderr is not None:
  278. filestderr.close()
  279. # Delete the temporary directory holding the test program and stderr files.
  280. if tempdir is not None:
  281. shutil.rmtree(tempdir)
  282. def strip_empty_entries(self, args):
  283. '''
  284. Drop any empty strings from the list of compile and link flags
  285. '''
  286. return [x for x in args if len(x) > 0]
  287. def cpu_arch(self):
  288. try:
  289. from cpuinfo import get_cpu_info
  290. except ImportError as e:
  291. cpu_info = self._backup_cpuinfo()
  292. if cpu_info is None:
  293. return "-march=native"
  294. try:
  295. cpu_info = get_cpu_info()
  296. except Exception as e:
  297. self.warning(
  298. f"{self.name} attempted to use `py-cpuinfo` but failed (exception type: {type(e)}, {e}), "
  299. "falling back to `lscpu` to get this information.")
  300. cpu_info = self._backup_cpuinfo()
  301. if cpu_info is None:
  302. return "-march=native"
  303. if cpu_info['arch'].startswith('PPC_'):
  304. # gcc does not provide -march on PowerPC, use -mcpu instead
  305. return '-mcpu=native'
  306. return '-march=native'
  307. def _backup_cpuinfo(self):
  308. # Construct cpu_info dict from lscpu that is similar to what py-cpuinfo provides
  309. if not self.command_exists('lscpu'):
  310. self.warning(
  311. f"{self.name} attempted to query 'lscpu' after failing to use py-cpuinfo "
  312. "to detect the CPU architecture. 'lscpu' does not appear to exist on "
  313. "your system, will fall back to use -march=native and non-vectorized execution."
  314. )
  315. return None
  316. result = subprocess.check_output('lscpu', shell=True)
  317. result = result.decode('utf-8').strip().lower()
  318. cpu_info = {}
  319. cpu_info['arch'] = None
  320. cpu_info['flags'] = ""
  321. if 'genuineintel' in result or 'authenticamd' in result:
  322. cpu_info['arch'] = 'X86_64'
  323. if 'avx512' in result:
  324. cpu_info['flags'] += 'avx512,'
  325. if 'avx2' in result:
  326. cpu_info['flags'] += 'avx2'
  327. elif 'ppc64le' in result:
  328. cpu_info['arch'] = "PPC_"
  329. return cpu_info
  330. def simd_width(self):
  331. try:
  332. from cpuinfo import get_cpu_info
  333. except ImportError as e:
  334. cpu_info = self._backup_cpuinfo()
  335. if cpu_info is None:
  336. return '-D__SCALAR__'
  337. try:
  338. cpu_info = get_cpu_info()
  339. except Exception as e:
  340. self.warning(
  341. f"{self.name} attempted to use `py-cpuinfo` but failed (exception type: {type(e)}, {e}), "
  342. "falling back to `lscpu` to get this information.")
  343. cpu_info = self._backup_cpuinfo()
  344. if cpu_info is None:
  345. return '-D__SCALAR__'
  346. if cpu_info['arch'] == 'X86_64':
  347. if 'avx512' in cpu_info['flags']:
  348. return '-D__AVX512__'
  349. elif 'avx2' in cpu_info['flags']:
  350. return '-D__AVX256__'
  351. return '-D__SCALAR__'
  352. def command_exists(self, cmd):
  353. if '|' in cmd:
  354. cmds = cmd.split("|")
  355. else:
  356. cmds = [cmd]
  357. valid = False
  358. for cmd in cmds:
  359. result = subprocess.Popen(f'type {cmd}', stdout=subprocess.PIPE, shell=True)
  360. valid = valid or result.wait() == 0
  361. if not valid and len(cmds) > 1:
  362. print(
  363. f"{WARNING} {self.name} requires one of the following commands '{cmds}', but it does not exist!"
  364. )
  365. elif not valid and len(cmds) == 1:
  366. print(
  367. f"{WARNING} {self.name} requires the '{cmd}' command, but it does not exist!"
  368. )
  369. return valid
  370. def warning(self, msg):
  371. self.error_log = f"{msg}"
  372. print(f"{WARNING} {msg}")
  373. def deepspeed_src_path(self, code_path):
  374. if os.path.isabs(code_path):
  375. return code_path
  376. else:
  377. return os.path.join(Path(__file__).parent.parent.absolute(), code_path)
  378. def builder(self):
  379. from torch.utils.cpp_extension import CppExtension
  380. return CppExtension(
  381. name=self.absolute_name(),
  382. sources=self.strip_empty_entries(self.sources()),
  383. include_dirs=self.strip_empty_entries(self.include_paths()),
  384. extra_compile_args={'cxx': self.strip_empty_entries(self.cxx_args())},
  385. extra_link_args=self.strip_empty_entries(self.extra_ldflags()))
  386. def load(self, verbose=True):
  387. from ...git_version_info import installed_ops, torch_info
  388. if installed_ops[self.name]:
  389. # Ensure the op we're about to load was compiled with the same
  390. # torch/cuda versions we are currently using at runtime.
  391. if isinstance(self, CUDAOpBuilder):
  392. self.assert_torch_info(torch_info)
  393. return importlib.import_module(self.absolute_name())
  394. else:
  395. return self.jit_load(verbose)
  396. def jit_load(self, verbose=True):
  397. if not self.is_compatible(verbose):
  398. raise RuntimeError(
  399. f"Unable to JIT load the {self.name} op due to it not being compatible due to hardware/software issue. {self.error_log}"
  400. )
  401. try:
  402. import ninja # noqa: F401
  403. except ImportError:
  404. raise RuntimeError(
  405. f"Unable to JIT load the {self.name} op due to ninja not being installed."
  406. )
  407. if isinstance(self, CUDAOpBuilder) and not self.is_rocm_pytorch():
  408. assert_no_cuda_mismatch()
  409. self.jit_mode = True
  410. from torch.utils.cpp_extension import load
  411. start_build = time.time()
  412. sources = [self.deepspeed_src_path(path) for path in self.sources()]
  413. extra_include_paths = [
  414. self.deepspeed_src_path(path) for path in self.include_paths()
  415. ]
  416. # Torch will try and apply whatever CCs are in the arch list at compile time,
  417. # we have already set the intended targets ourselves we know that will be
  418. # needed at runtime. This prevents CC collisions such as multiple __half
  419. # implementations. Stash arch list to reset after build.
  420. torch_arch_list = None
  421. if "TORCH_CUDA_ARCH_LIST" in os.environ:
  422. torch_arch_list = os.environ.get("TORCH_CUDA_ARCH_LIST")
  423. os.environ["TORCH_CUDA_ARCH_LIST"] = ""
  424. op_module = load(
  425. name=self.name,
  426. sources=self.strip_empty_entries(sources),
  427. extra_include_paths=self.strip_empty_entries(extra_include_paths),
  428. extra_cflags=self.strip_empty_entries(self.cxx_args()),
  429. extra_cuda_cflags=self.strip_empty_entries(self.nvcc_args()),
  430. extra_ldflags=self.strip_empty_entries(self.extra_ldflags()),
  431. verbose=verbose)
  432. build_duration = time.time() - start_build
  433. if verbose:
  434. print(f"Time to load {self.name} op: {build_duration} seconds")
  435. # Reset arch list so we are not silently removing it for other possible use cases
  436. if torch_arch_list:
  437. os.environ["TORCH_CUDA_ARCH_LIST"] = torch_arch_list
  438. return op_module
  439. class CUDAOpBuilder(OpBuilder):
  440. def compute_capability_args(self, cross_compile_archs=None):
  441. """
  442. Returns nvcc compute capability compile flags.
  443. 1. `TORCH_CUDA_ARCH_LIST` takes priority over `cross_compile_archs`.
  444. 2. If neither is set default compute capabilities will be used
  445. 3. Under `jit_mode` compute capabilities of all visible cards will be used plus PTX
  446. Format:
  447. - `TORCH_CUDA_ARCH_LIST` may use ; or whitespace separators. Examples:
  448. TORCH_CUDA_ARCH_LIST="6.1;7.5;8.6" pip install ...
  449. TORCH_CUDA_ARCH_LIST="6.0 6.1 7.0 7.5 8.0 8.6+PTX" pip install ...
  450. - `cross_compile_archs` uses ; separator.
  451. """
  452. ccs = []
  453. if self.jit_mode:
  454. # Compile for underlying architectures since we know those at runtime
  455. for i in range(torch.cuda.device_count()):
  456. CC_MAJOR, CC_MINOR = torch.cuda.get_device_capability(i)
  457. cc = f"{CC_MAJOR}.{CC_MINOR}"
  458. if cc not in ccs:
  459. ccs.append(cc)
  460. ccs = sorted(ccs)
  461. ccs[-1] += '+PTX'
  462. else:
  463. # Cross-compile mode, compile for various architectures
  464. # env override takes priority
  465. cross_compile_archs_env = os.environ.get('TORCH_CUDA_ARCH_LIST', None)
  466. if cross_compile_archs_env is not None:
  467. if cross_compile_archs is not None:
  468. print(
  469. f"{WARNING} env var `TORCH_CUDA_ARCH_LIST={cross_compile_archs_env}` overrides `cross_compile_archs={cross_compile_archs}`"
  470. )
  471. cross_compile_archs = cross_compile_archs_env.replace(' ', ';')
  472. else:
  473. if cross_compile_archs is None:
  474. cross_compile_archs = get_default_compute_capabilities()
  475. ccs = cross_compile_archs.split(';')
  476. ccs = self.filter_ccs(ccs)
  477. if len(ccs) == 0:
  478. raise RuntimeError(
  479. f"Unable to load {self.name} op due to no compute capabilities remaining after filtering"
  480. )
  481. args = []
  482. for cc in ccs:
  483. num = cc[0] + cc[2]
  484. args.append(f'-gencode=arch=compute_{num},code=sm_{num}')
  485. if cc.endswith('+PTX'):
  486. args.append(f'-gencode=arch=compute_{num},code=compute_{num}')
  487. return args
  488. def filter_ccs(self, ccs: List[str]):
  489. """
  490. Prune any compute capabilities that are not compatible with the builder. Should log
  491. which CCs have been pruned.
  492. """
  493. return ccs
  494. def version_dependent_macros(self):
  495. # Fix from apex that might be relevant for us as well, related to https://github.com/NVIDIA/apex/issues/456
  496. version_ge_1_1 = []
  497. if (TORCH_MAJOR > 1) or (TORCH_MAJOR == 1 and TORCH_MINOR > 0):
  498. version_ge_1_1 = ['-DVERSION_GE_1_1']
  499. version_ge_1_3 = []
  500. if (TORCH_MAJOR > 1) or (TORCH_MAJOR == 1 and TORCH_MINOR > 2):
  501. version_ge_1_3 = ['-DVERSION_GE_1_3']
  502. version_ge_1_5 = []
  503. if (TORCH_MAJOR > 1) or (TORCH_MAJOR == 1 and TORCH_MINOR > 4):
  504. version_ge_1_5 = ['-DVERSION_GE_1_5']
  505. return version_ge_1_1 + version_ge_1_3 + version_ge_1_5
  506. def is_compatible(self, verbose=True):
  507. return super().is_compatible(verbose)
  508. def builder(self):
  509. from torch.utils.cpp_extension import CUDAExtension
  510. if not self.is_rocm_pytorch():
  511. assert_no_cuda_mismatch()
  512. cuda_ext = CUDAExtension(
  513. name=self.absolute_name(),
  514. sources=self.strip_empty_entries(self.sources()),
  515. include_dirs=self.strip_empty_entries(self.include_paths()),
  516. libraries=self.strip_empty_entries(self.libraries_args()),
  517. extra_compile_args={
  518. 'cxx': self.strip_empty_entries(self.cxx_args()),
  519. 'nvcc': self.strip_empty_entries(self.nvcc_args())
  520. })
  521. if self.is_rocm_pytorch():
  522. # hip converts paths to absolute, this converts back to relative
  523. sources = cuda_ext.sources
  524. curr_file = Path(__file__).parent.parent # ds root
  525. for i in range(len(sources)):
  526. src = Path(sources[i])
  527. sources[i] = str(src.relative_to(curr_file))
  528. cuda_ext.sources = sources
  529. return cuda_ext
  530. def hipify_extension(self):
  531. if self.is_rocm_pytorch():
  532. from torch.utils.hipify import hipify_python
  533. hipify_python.hipify(
  534. project_directory=os.getcwd(),
  535. output_directory=os.getcwd(),
  536. header_include_dirs=self.include_paths(),
  537. includes=[os.path.join(os.getcwd(),
  538. '*')],
  539. extra_files=[os.path.abspath(s) for s in self.sources()],
  540. show_detailed=True,
  541. is_pytorch_extension=True,
  542. hipify_extra_files_only=True,
  543. )
  544. def cxx_args(self):
  545. if sys.platform == "win32":
  546. return ['-O2']
  547. else:
  548. return ['-O3', '-std=c++14', '-g', '-Wno-reorder']
  549. def nvcc_args(self):
  550. args = ['-O3']
  551. if self.is_rocm_pytorch():
  552. ROCM_MAJOR, ROCM_MINOR = self.installed_rocm_version()
  553. args += [
  554. '-std=c++14',
  555. '-U__HIP_NO_HALF_OPERATORS__',
  556. '-U__HIP_NO_HALF_CONVERSIONS__',
  557. '-U__HIP_NO_HALF2_OPERATORS__',
  558. '-DROCM_VERSION_MAJOR=%s' % ROCM_MAJOR,
  559. '-DROCM_VERSION_MINOR=%s' % ROCM_MINOR
  560. ]
  561. else:
  562. cuda_major, _ = installed_cuda_version()
  563. args += [
  564. '-allow-unsupported-compiler' if sys.platform == "win32" else '',
  565. '--use_fast_math',
  566. '-std=c++17'
  567. if sys.platform == "win32" and cuda_major > 10 else '-std=c++14',
  568. '-U__CUDA_NO_HALF_OPERATORS__',
  569. '-U__CUDA_NO_HALF_CONVERSIONS__',
  570. '-U__CUDA_NO_HALF2_OPERATORS__'
  571. ]
  572. args += self.compute_capability_args()
  573. return args
  574. def libraries_args(self):
  575. if sys.platform == "win32":
  576. return ['cublas', 'curand']
  577. else:
  578. return []
  579. class TorchCPUOpBuilder(CUDAOpBuilder):
  580. def extra_ldflags(self):
  581. if not self.is_rocm_pytorch():
  582. return ['-lcurand']
  583. else:
  584. return []
  585. def cxx_args(self):
  586. import torch
  587. if not self.is_rocm_pytorch():
  588. CUDA_LIB64 = os.path.join(torch.utils.cpp_extension.CUDA_HOME, "lib64")
  589. else:
  590. CUDA_LIB64 = os.path.join(torch.utils.cpp_extension.ROCM_HOME, "lib")
  591. CPU_ARCH = self.cpu_arch()
  592. SIMD_WIDTH = self.simd_width()
  593. args = super().cxx_args()
  594. args += [
  595. f'-L{CUDA_LIB64}',
  596. '-lcudart',
  597. '-lcublas',
  598. '-g',
  599. CPU_ARCH,
  600. '-fopenmp',
  601. SIMD_WIDTH,
  602. ]
  603. return args