{ "cells": [ { "cell_type": "markdown", "metadata": {}, "source": [ "(hpu_bert_training)=\n", "# BERT Model Training with Intel Gaudi\n", "\n", "\n", " \"try-anyscale-quickstart\"\n", "\n", "

\n", "\n", "In this notebook, we will train a BERT model for sequence classification using the Yelp review full dataset. We will use the `transformers` and `datasets` libraries from Hugging Face, along with `ray.train` for distributed training.\n", "\n", "[Intel Gaudi AI Processors (HPUs)](https://habana.ai) are AI hardware accelerators designed by Intel Habana Labs. For more information, see [Gaudi Architecture](https://docs.habana.ai/en/latest/Gaudi_Overview/index.html) and [Gaudi Developer Docs](https://developer.habana.ai/).\n", "\n", "## Configuration\n", "\n", "A node with Gaudi/Gaudi2 installed is required to run this example. Both Gaudi and Gaudi2 have 8 HPUs. We will use 2 workers to train the model, each using 1 HPU.\n", "\n", "We recommend using a prebuilt container to run these examples. To run a container, you need Docker. See [Install Docker Engine](https://docs.docker.com/engine/install/) for installation instructions.\n", "\n", "Next, follow [Run Using Containers](https://docs.habana.ai/en/latest/Installation_Guide/Bare_Metal_Fresh_OS.html?highlight=installer#run-using-containers) to install the Gaudi drivers and container runtime.\n", "\n", "Next, start the Gaudi container:\n", "```bash\n", "docker pull vault.habana.ai/gaudi-docker/1.20.0/ubuntu22.04/habanalabs/pytorch-installer-2.6.0:latest\n", "docker run -it --runtime=habana -e HABANA_VISIBLE_DEVICES=all -e OMPI_MCA_btl_vader_single_copy_mechanism=none --cap-add=sys_nice --net=host --ipc=host vault.habana.ai/gaudi-docker/1.20.0/ubuntu22.04/habanalabs/pytorch-installer-2.6.0:latest\n", "```\n", "\n", "Inside the container, install the following dependecies to run this notebook.\n", "```bash\n", "pip install ray[train] notebook transformers datasets evaluate\n", "```" ] }, { "cell_type": "code", "execution_count": 1, "metadata": {}, "outputs": [ { "name": "stderr", "output_type": "stream", "text": [ "/usr/local/lib/python3.10/dist-packages/torch/distributed/distributed_c10d.py:252: UserWarning: Device capability of hccl unspecified, assuming `cpu` and `cuda`. Please specify it via the `devices` argument of `register_backend`.\n", " warnings.warn(\n" ] } ], "source": [ "# Import necessary libraries\n", "\n", "import os\n", "from typing import Dict\n", "\n", "import torch\n", "from torch import nn\n", "from torch.utils.data import DataLoader\n", "from tqdm import tqdm\n", "\n", "import numpy as np\n", "import evaluate\n", "from datasets import load_dataset\n", "import transformers\n", "from transformers import (\n", " Trainer,\n", " TrainingArguments,\n", " AutoTokenizer,\n", " AutoModelForSequenceClassification,\n", ")\n", "\n", "import ray.train\n", "from ray.train import ScalingConfig\n", "from ray.train.torch import TorchTrainer\n", "from ray.train.torch import TorchConfig\n", "from ray.runtime_env import RuntimeEnv\n", "\n", "import habana_frameworks.torch.core as htcore" ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "## Metrics Setup\n", "\n", "We will use accuracy as our evaluation metric. The `compute_metrics` function will calculate the accuracy of our model's predictions." ] }, { "cell_type": "code", "execution_count": 2, "metadata": {}, "outputs": [], "source": [ "# Metrics\n", "metric = evaluate.load(\"accuracy\")\n", "\n", "def compute_metrics(eval_pred):\n", " logits, labels = eval_pred\n", " predictions = np.argmax(logits, axis=-1)\n", " return metric.compute(predictions=predictions, references=labels)" ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "## Training Function\n", "\n", "This function will be executed by each worker during training. It handles data loading, tokenization, model initialization, and the training loop. Compared to a training function for GPU, no changes are needed to port to HPU. Internally, Ray Train does these things:\n", "\n", "* Detect HPU and set the device.\n", "\n", "* Initializes the habana PyTorch backend.\n", "\n", "* Initializes the habana distributed backend." ] }, { "cell_type": "code", "execution_count": 3, "metadata": {}, "outputs": [], "source": [ "def train_func_per_worker(config: Dict):\n", " \n", " # Datasets\n", " dataset = load_dataset(\"yelp_review_full\")\n", " tokenizer = AutoTokenizer.from_pretrained(\"bert-base-cased\")\n", " \n", " def tokenize_function(examples):\n", " return tokenizer(examples[\"text\"], padding=\"max_length\", truncation=True)\n", "\n", " lr = config[\"lr\"]\n", " epochs = config[\"epochs\"]\n", " batch_size = config[\"batch_size_per_worker\"]\n", "\n", " train_dataset = dataset[\"train\"].select(range(1000)).map(tokenize_function, batched=True)\n", " eval_dataset = dataset[\"test\"].select(range(1000)).map(tokenize_function, batched=True)\n", "\n", " # Prepare dataloader for each worker\n", " dataloaders = {}\n", " dataloaders[\"train\"] = torch.utils.data.DataLoader(\n", " train_dataset, \n", " shuffle=True, \n", " collate_fn=transformers.default_data_collator, \n", " batch_size=batch_size\n", " )\n", " dataloaders[\"test\"] = torch.utils.data.DataLoader(\n", " eval_dataset, \n", " shuffle=True, \n", " collate_fn=transformers.default_data_collator, \n", " batch_size=batch_size\n", " )\n", "\n", " # Obtain HPU device automatically\n", " device = ray.train.torch.get_device()\n", "\n", " # Prepare model and optimizer\n", " model = AutoModelForSequenceClassification.from_pretrained(\n", " \"bert-base-cased\", num_labels=5\n", " )\n", " model = model.to(device)\n", " \n", " optimizer = torch.optim.SGD(model.parameters(), lr=lr, momentum=0.9)\n", "\n", " # Start training loops\n", " for epoch in range(epochs):\n", " # Each epoch has a training and validation phase\n", " for phase in [\"train\", \"test\"]:\n", " if phase == \"train\":\n", " model.train() # Set model to training mode\n", " else:\n", " model.eval() # Set model to evaluate mode\n", "\n", " # breakpoint()\n", " for batch in dataloaders[phase]:\n", " batch = {k: v.to(device) for k, v in batch.items()}\n", "\n", " # zero the parameter gradients\n", " optimizer.zero_grad()\n", "\n", " # forward\n", " with torch.set_grad_enabled(phase == \"train\"):\n", " # Get model outputs and calculate loss\n", " \n", " outputs = model(**batch)\n", " loss = outputs.loss\n", "\n", " # backward + optimize only if in training phase\n", " if phase == \"train\":\n", " loss.backward()\n", " optimizer.step()\n", " print(f\"train epoch:[{epoch}]\\tloss:{loss:.6f}\")" ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "## Main Training Function\n", "\n", "The `train_bert` function sets up the distributed training environment using Ray and starts the training process. To enable training using HPU, we only need to make the following changes:\n", "* Require an HPU for each worker in ScalingConfig\n", "* Set backend to \"hccl\" in TorchConfig" ] }, { "cell_type": "code", "execution_count": 4, "metadata": {}, "outputs": [], "source": [ "def train_bert(num_workers=2):\n", " global_batch_size = 8\n", "\n", " train_config = {\n", " \"lr\": 1e-3,\n", " \"epochs\": 10,\n", " \"batch_size_per_worker\": global_batch_size // num_workers,\n", " }\n", "\n", " # Configure computation resources\n", " # In ScalingConfig, require an HPU for each worker\n", " scaling_config = ScalingConfig(num_workers=num_workers, resources_per_worker={\"CPU\": 1, \"HPU\": 1})\n", " # Set backend to hccl in TorchConfig\n", " torch_config = TorchConfig(backend = \"hccl\")\n", " \n", " # start your ray cluster\n", " ray.init()\n", " \n", " # Initialize a Ray TorchTrainer\n", " trainer = TorchTrainer(\n", " train_loop_per_worker=train_func_per_worker,\n", " train_loop_config=train_config,\n", " torch_config=torch_config,\n", " scaling_config=scaling_config,\n", " )\n", "\n", " result = trainer.fit()\n", " print(f\"Training result: {result}\")" ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "## Start Training\n", "\n", "Finally, we call the `train_bert` function to start the training process. You can adjust the number of workers to use.\n", "\n", "Note: the following warning is fine, and is resolved in SynapseAI version 1.14.0+:\n", "```text\n", "/usr/local/lib/python3.10/dist-packages/torch/distributed/distributed_c10d.py:252: UserWarning: Device capability of hccl unspecified, assuming `cpu` and `cuda`. Please specify it via the `devices` argument of `register_backend`.\n", "```" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "train_bert(num_workers=2)" ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "## Possible outputs\n", "\n", "``` text\n", "Downloading builder script: 100%|██████████| 4.20k/4.20k [00:00<00:00, 27.0MB/s]\n", "2025-03-03 03:37:08,776 INFO worker.py:1841 -- Started a local Ray instance.\n", "/usr/local/lib/python3.10/dist-packages/ray/tune/impl/tuner_internal.py:125: RayDeprecationWarning: The `RunConfig` class should be imported from `ray.tune` when passing it to the Tuner. Please update your imports. See this issue for more context and migration options: https://github.com/ray-project/ray/issues/49454. Disable these warnings by setting the environment variable: RAY_TRAIN_ENABLE_V2_MIGRATION_WARNINGS=0\n", " _log_deprecation_warning(\n", "(RayTrainWorker pid=75123) Setting up process group for: env:// [rank=0, world_size=2]\n", "(TorchTrainer pid=74734) Started distributed worker processes: \n", "(TorchTrainer pid=74734) - (node_id=eef984cd0cd96cce50bad1b1dab12e19c809047f10be3c829524a3d1, ip=100.83.111.228, pid=75123) world_rank=0, local_rank=0, node_rank=0\n", "(TorchTrainer pid=74734) - (node_id=eef984cd0cd96cce50bad1b1dab12e19c809047f10be3c829524a3d1, ip=100.83.111.228, pid=75122) world_rank=1, local_rank=1, node_rank=0\n", "Generating train split: 0%| | 0/650000 [00:00