notebooks/multilabel_HoC.ipynb (1,293 lines of code) (raw):

{ "cells": [ { "attachments": {}, "cell_type": "markdown", "metadata": { "id": "4IYf85KnrWoH" }, "source": [ "# SetFit SOTA for Bio Text Classification" ] }, { "attachments": {}, "cell_type": "markdown", "metadata": { "id": "DM_RojZL3XDF" }, "source": [ "SetFit is a great practical tool for few shot text classification, but did you know that you can fine-tune a vanilla SetFit for full-shot text classification and outperform models that were pre-trained from scratch using domain data?\n", "Here we show such example in the Biological domain, where SetFit outperforms most of the models that were trained from scratch on Biological data, while being more efficient." ] }, { "attachments": {}, "cell_type": "markdown", "metadata": { "id": "oS_qaldP3Ths" }, "source": [ "The following table summarizes the results of different models on the HoC* dataset. All of the biological models were first pre-trained using in-domain biological data and in addition were fine-tuned given the HoC training data in the BLUE benchmark. SetFit was not pre-trained using biological data, it is based on a general pre-trained sentence transformer model (MSFT's mpnet) and was solely fine-tuned on the HoC training data. As shown in the table, SetFit surpasses the Bio models and achieves comparable performance to the 347M BioGPT, which is the SOTA model for the Bio domain, while being 3x smaller: https://analyticsindiamag.com/microsoft-launches-biogpt-the-chatgpt-of-lifescience/" ] }, { "cell_type": "markdown", "metadata": { "id": "aLPGtW9B3I2J" }, "source": [ "| **Model** | **#params[M]** | **F1** | **Pre-train Data** | \n", "|:-----------------------:|:-------:|:---------------:|:-----------------:|\n", "| **BioBERT[1]**| 110 | 81.5 | Bio \n", "| **PubMedBERT[2]**| 340 | 82.7 | Bio \n", "| **BioLinkBERT[3]** | 340 | 84.9 | Bio \n", "| **GPT-2** | 355 | 81.8 | General \n", "| **BioGPT[4]** | 347 | 85.1 | Bio\n", "| **SetFit** | 105 | **85.1** | General\n", "\n", "\n" ] }, { "cell_type": "markdown", "metadata": { "id": "arGWgoRK3BlE" }, "source": [ "Refrences:\n", "\n", "[1] Domain-specific\n", "language model pretraining for biomedical natural language\n", "processing\" https://arxiv.org/abs/2007.15779\n", "\n", "[2] BioBERT: a pre-trained biomedical language representation\n", "model for biomedical text mining\" https://arxiv.org/abs/1901.08746\n", "\n", "[3] LinkBERT: Pretraining Language Models with Document Links https://arxiv.org/abs/2203.15827\n", "\n", "[4] BioGPT: Generative Pre-trained Transformer for Biomedical Text Generation and Mining\" https://arxiv.org/abs/2210.10341\n", "\n", "[5] Automatic semantic classification of scientific literature according to\n", "the hallmarks of cancer. https://academic.oup.com/bioinformatics/article/32/3/432/1743783\n", "\n", "[6] An\n", "evaluation of BERT and ELMo on ten benchmarking\n", "datasets https://arxiv.org/abs/1906.05474\n" ] }, { "cell_type": "markdown", "metadata": { "id": "N_lI2PxcPOAJ" }, "source": [ "*HoC (the Hallmarks of Cancers corpus) consists of 1580\n", "PubMed abstracts manually annotated at sentence level by\n", "experts with ten currently known hallmarks of cancer [5]. We follow the same training/test split as in [6]" ] }, { "attachments": {}, "cell_type": "markdown", "metadata": { "id": "zelvDdXVvV4a" }, "source": [ "### SetFit Multilabel HoC" ] }, { "cell_type": "code", "execution_count": null, "metadata": { "id": "DaL_TsPWrcI3" }, "outputs": [], "source": [ "!pip install setfit" ] }, { "attachments": {}, "cell_type": "markdown", "metadata": { "id": "etBVpBqJrpd-" }, "source": [ "Load the HoC dataset" ] }, { "cell_type": "code", "execution_count": null, "metadata": { "id": "wTZgoQPPrn7a" }, "outputs": [], "source": [ "!wget https://github.com/ncbi-nlp/BLUE_Benchmark/releases/download/0.1/data_v0.1.zip\n", "!unzip data_v0.1.zip\n", "\n", "import pandas as pd\n", "import numpy as np\n", "\n", "# Read train/test files\n", "test_df = pd.read_csv('/content/data/hoc/test.tsv', sep='\\t')\n", "train_df = pd.read_csv('/content/data/hoc/train.tsv', sep='\\t')" ] }, { "cell_type": "code", "execution_count": null, "metadata": { "id": "WbcTmX31sqvY" }, "outputs": [], "source": [ "LABELS = ['activating invasion and metastasis', 'avoiding immune destruction',\n", " 'cellular energetics', 'enabling replicative immortality', 'evading growth suppressors',\n", " 'genomic instability and mutation', 'inducing angiogenesis', 'resisting cell death',\n", " 'sustaining proliferative signaling', 'tumor promoting inflammation']" ] }, { "cell_type": "code", "execution_count": null, "metadata": { "id": "kvlT2j6zwfGM" }, "outputs": [], "source": [ "# Convert labels to hotvec multilabel format (similar to scikit-learn)\n", "def hotvec_multilabel(true_df):\n", " data = {}\n", "\n", " for i in range(len(true_df)):\n", " true_row = true_df.iloc[i]\n", "\n", " key = true_row['index']\n", "\n", " data[key] = set()\n", "\n", " if not pd.isna(true_row['labels']):\n", " for l in true_row['labels'].split(','):\n", " data[key].add(LABELS.index(l))\n", " \n", " y_hotvec = []\n", " for k, (true) in data.items():\n", " t = [0] * len(LABELS)\n", " for i in true:\n", " t[i] = 1\n", "\n", " y_hotvec.append(t)\n", "\n", " y_hotvec = np.array(y_hotvec)\n", "\n", " return y_hotvec" ] }, { "attachments": {}, "cell_type": "markdown", "metadata": { "id": "OzN7ATsqsuFn" }, "source": [ "### SetFit Multilabel" ] }, { "cell_type": "code", "execution_count": null, "metadata": { "colab": { "base_uri": "https://localhost:8080/" }, "id": "Pbu3EwqNsx4W", "outputId": "ac1b441d-1f25-4a9b-dfb9-722ccf0898f3" }, "outputs": [ { "name": "stderr", "output_type": "stream", "text": [ "model_head.pkl not found on HuggingFace Hub, initialising classification head with random weights. You should TRAIN this model on a downstream task to use it for predictions and inference.\n" ] } ], "source": [ "from datasets import Dataset\n", "import evaluate\n", "from setfit import SetFitModel, SetFitTrainer\n", "\n", "model = SetFitModel.from_pretrained(\n", " \"sentence-transformers/paraphrase-mpnet-base-v2\", \n", " multi_target_strategy=\"multi-output\", # one-vs-rest; multi-output; classifier-chain\n", ")\n", "\n", "multilabel_f1_metric = evaluate.load(\"f1\", \"multilabel\")\n", "multilabel_accuracy_metric = evaluate.load(\"accuracy\", \"multilabel\")\n", "\n", "# f1/accuracy sentence level\n", "def compute_metrics(y_pred, y_test):\n", " return {\n", " \"f1\": multilabel_f1_metric.compute(predictions=y_pred, references=y_test, average=\"micro\")[\"f1\"],\n", " \"accuracy\": multilabel_accuracy_metric.compute(predictions=y_pred, references=y_test)[\"accuracy\"],\n", " }" ] }, { "cell_type": "code", "execution_count": null, "metadata": { "id": "Atx4T90ts4L5" }, "outputs": [], "source": [ "eval_dataset = Dataset.from_dict({\"text\": test_df['sentence'], \"label\": hotvec_multilabel(test_df)})\n", "train_dataset = Dataset.from_dict({\"text\": train_df['sentence'], \"label\": hotvec_multilabel(train_df)})" ] }, { "cell_type": "code", "execution_count": null, "metadata": { "id": "m4wzUbmLtIJr" }, "outputs": [], "source": [ "trainer = SetFitTrainer(\n", " model=model,\n", " train_dataset=train_dataset,\n", " eval_dataset=eval_dataset,\n", " metric=compute_metrics,\n", " num_iterations=5,\n", ")" ] }, { "cell_type": "code", "execution_count": null, "metadata": { "colab": { "base_uri": "https://localhost:8080/", "height": 515, "referenced_widgets": [ "4180378f634c4fd0a3c5013159cd6770", "7c67c459c3c74f3eafe253e4865bed32", "a31137fa736243509f27b08b0163630b", "9bbfab5bf2354264879c63959d608daf", "23f77655895c476d94cb02f152d0e4b7", "88b1e64d0a8149ebb81316c419c36a47", "83ebfefc9345405bb0224c8df59e7d83", "1b65756da31a43dd9ca4f49da6640e73", "08ef465f79924616b29fe9828089d830", "8f1c94ebd5a44c1ca9bcb9cfe8a85d56", "31c15b576f134c2b8d6d14ca77dd2f88", "7f933d6a976f4542bee79cf6bf1fe497", "0ee3154796c44326989a4db3a3331293", "df924c28c2214ab48c46999f8e332182", "b39a63cc65224a8698cc2fa06af12ead", "ed149912957740a393a0029497d4a144", "e191f4a2517741689d2d11800da5cc45", "f3436e0825204faca0a94f6795d75e89", "cc631484771647b989a24564792f29e4", "3a737d671bfb4aa6aee609a80ddede72", "d0202adacfa64dc6aeb0a3f4c33e28fc", "153550864ab84636af528e7eb8a959a3" ] }, "id": "mJu5u-vltPMz", "outputId": "423499ee-758e-45b4-96a8-ebba98adafd2" }, "outputs": [ { "name": "stderr", "output_type": "stream", "text": [ "Applying column mapping to training dataset\n", "***** Running training *****\n", " Num examples = 71200\n", " Num epochs = 1\n", " Total optimization steps = 4450\n", " Total train batch size = 16\n" ] }, { "data": { "application/vnd.jupyter.widget-view+json": { "model_id": "4180378f634c4fd0a3c5013159cd6770", "version_major": 2, "version_minor": 0 }, "text/plain": [ "Epoch: 0%| | 0/1 [00:00<?, ?it/s]" ] }, "metadata": {}, "output_type": "display_data" }, { "data": { "application/vnd.jupyter.widget-view+json": { "model_id": "7f933d6a976f4542bee79cf6bf1fe497", "version_major": 2, "version_minor": 0 }, "text/plain": [ "Iteration: 0%| | 0/4450 [00:00<?, ?it/s]" ] }, "metadata": {}, "output_type": "display_data" }, { "name": "stderr", "output_type": "stream", "text": [ "/usr/local/lib/python3.9/dist-packages/sklearn/linear_model/_logistic.py:458: ConvergenceWarning: lbfgs failed to converge (status=1):\n", "STOP: TOTAL NO. of ITERATIONS REACHED LIMIT.\n", "\n", "Increase the number of iterations (max_iter) or scale the data as shown in:\n", " https://scikit-learn.org/stable/modules/preprocessing.html\n", "Please also refer to the documentation for alternative solver options:\n", " https://scikit-learn.org/stable/modules/linear_model.html#logistic-regression\n", " n_iter_i = _check_optimize_result(\n", "/usr/local/lib/python3.9/dist-packages/sklearn/linear_model/_logistic.py:458: ConvergenceWarning: lbfgs failed to converge (status=1):\n", "STOP: TOTAL NO. of ITERATIONS REACHED LIMIT.\n", "\n", "Increase the number of iterations (max_iter) or scale the data as shown in:\n", " https://scikit-learn.org/stable/modules/preprocessing.html\n", "Please also refer to the documentation for alternative solver options:\n", " https://scikit-learn.org/stable/modules/linear_model.html#logistic-regression\n", " n_iter_i = _check_optimize_result(\n", "Applying column mapping to evaluation dataset\n", "***** Running evaluation *****\n" ] }, { "name": "stdout", "output_type": "stream", "text": [ "{'f1': 0.7284569138276554, 'accuracy': 0.836671270718232}\n" ] } ], "source": [ "trainer.train()\n", "metrics = trainer.evaluate()\n", "print(metrics)" ] }, { "attachments": {}, "cell_type": "markdown", "metadata": { "id": "409BjocBtkkf" }, "source": [ "### Evaluation of BLUE's HoC F1 (abstract level)" ] }, { "cell_type": "markdown", "metadata": { "id": "ifJTpGsItsXJ" }, "source": [ "Support functions refactored from https://github.com/ncbi-nlp/BLUE_Benchmark\n", "can be downloaded at https://github.com/ncbi-nlp/BLUE_Benchmark/releases/tag/0.1" ] }, { "cell_type": "code", "execution_count": null, "metadata": { "id": "hbNss08Wtkkh" }, "outputs": [], "source": [ "def divide(x, y):\n", " return np.true_divide(x, y, out=np.zeros_like(x, dtype=np.float64), where=y != 0)\n", "\n", "def get_p_r_f_arrary(test_predict_label, test_true_label):\n", " num, cat = test_predict_label.shape\n", " acc_list = []\n", " prc_list = []\n", " rec_list = []\n", " f_score_list = []\n", " for i in range(num):\n", " label_pred_set = set()\n", " label_gold_set = set()\n", "\n", " for j in range(cat):\n", " if test_predict_label[i, j] == 1:\n", " label_pred_set.add(j)\n", " if test_true_label[i, j] == 1:\n", " label_gold_set.add(j)\n", "\n", " uni_set = label_gold_set.union(label_pred_set)\n", " intersec_set = label_gold_set.intersection(label_pred_set)\n", "\n", " tt = len(intersec_set)\n", " if len(label_pred_set) == 0:\n", " prc = 0\n", " else:\n", " prc = tt / len(label_pred_set)\n", "\n", " acc = tt / len(uni_set)\n", "\n", " rec = tt / len(label_gold_set)\n", "\n", " if prc == 0 and rec == 0:\n", " f_score = 0\n", " else:\n", " f_score = 2 * prc * rec / (prc + rec)\n", "\n", " acc_list.append(acc)\n", " prc_list.append(prc)\n", " rec_list.append(rec)\n", " f_score_list.append(f_score)\n", "\n", " mean_prc = np.mean(prc_list)\n", " mean_rec = np.mean(rec_list)\n", " f_score = divide(2 * mean_prc * mean_rec, (mean_prc + mean_rec))\n", " return mean_prc, mean_rec, f_score\n", "\n", "def eval_hoc(true_df, pred_df):\n", " data = {}\n", "\n", " assert len(true_df) == len(pred_df), \\\n", " f'Gold line no {len(true_df)} vs Prediction line no {len(pred_df)}'\n", "\n", " for i in range(len(true_df)):\n", " true_row = true_df.iloc[i]\n", " pred_row = pred_df.iloc[i]\n", " assert true_row['index'] == pred_row['index'], \\\n", " 'Index does not match @{}: {} vs {}'.format(i, true_row['index'], pred_row['index'])\n", "\n", " key = true_row['index'][:true_row['index'].find('_')]\n", " if key not in data:\n", " data[key] = (set(), set())\n", "\n", " if not pd.isna(true_row['labels']):\n", " for l in true_row['labels'].split(','):\n", " data[key][0].add(LABELS.index(l))\n", "\n", " if not pd.isna(pred_row['labels']):\n", " for l in pred_row['labels'].split(','):\n", " data[key][1].add(LABELS.index(l))\n", "\n", " assert len(data) == 315, 'There are 315 documents in the test set: %d' % len(data)\n", "\n", " y_test = []\n", " y_pred = []\n", " for k, (true, pred) in data.items():\n", " t = [0] * len(LABELS)\n", " for i in true:\n", " t[i] = 1\n", "\n", " p = [0] * len(LABELS)\n", " for i in pred:\n", " p[i] = 1\n", "\n", " y_test.append(t)\n", " y_pred.append(p)\n", "\n", " y_test = np.array(y_test)\n", " y_pred = np.array(y_pred)\n", "\n", " r, p, f1 = get_p_r_f_arrary(y_pred, y_test)\n", " print('Precision: {:.1f}'.format(p*100))\n", " print('Recall : {:.1f}'.format(r*100))\n", " print('F1 : {:.1f}'.format(f1*100))" ] }, { "attachments": {}, "cell_type": "markdown", "metadata": { "id": "PE_A3qdbtxns" }, "source": [ "#### Evaluate on test data" ] }, { "cell_type": "code", "execution_count": null, "metadata": { "id": "0A6rzsLKt5wC" }, "outputs": [], "source": [ "test_predict_label = trainer.model.predict(test_df['sentence'])" ] }, { "cell_type": "code", "execution_count": null, "metadata": { "id": "ya4Ps34ut8Dm" }, "outputs": [], "source": [ "# Convert hotvec multilabel to actual labels \n", "num, cat = test_predict_label.shape\n", "sentence_list = []\n", "for i in range(num):\n", " sentence_set = set()\n", " for j in range(cat):\n", " if test_predict_label[i, j] == 1:\n", " sentence_set.add(LABELS[j])\n", " sentence_list.append(','.join(sentence_set))\n", "\n", "# Reformat for HoC evaluation\n", "pred_df = test_df\n", "pred_df = pred_df.assign(labels = sentence_list)\n", "pred_df['labels'] = pred_df['labels'].replace({'':np.nan})\n", "test_df['labels'] = test_df['labels'].replace({'':np.nan})" ] }, { "attachments": {}, "cell_type": "markdown", "metadata": { "id": "B_5Typrvut_z" }, "source": [ "#### Evaluate F1 (abstract level)" ] }, { "cell_type": "code", "execution_count": null, "metadata": { "colab": { "base_uri": "https://localhost:8080/" }, "id": "5ZdQNVDPuyDJ", "outputId": "68d2de8b-eea2-4b70-fea6-e28d1998f75c" }, "outputs": [ { "name": "stdout", "output_type": "stream", "text": [ "Precision: 86.4\n", "Recall : 83.8\n", "F1 : 85.1\n" ] } ], "source": [ "eval_hoc(test_df, pred_df)" ] } ], "metadata": { "accelerator": "GPU", "colab": { "provenance": [] }, "gpuClass": "standard", "kernelspec": { "display_name": "Python 3", "name": "python3" }, "language_info": { "name": "python" }, "widgets": { "application/vnd.jupyter.widget-state+json": { "08ef465f79924616b29fe9828089d830": { "model_module": "@jupyter-widgets/controls", "model_module_version": "1.5.0", "model_name": "ProgressStyleModel", "state": { "_model_module": "@jupyter-widgets/controls", "_model_module_version": "1.5.0", "_model_name": "ProgressStyleModel", "_view_count": null, "_view_module": "@jupyter-widgets/base", "_view_module_version": "1.2.0", "_view_name": "StyleView", "bar_color": null, "description_width": "" } }, "0ee3154796c44326989a4db3a3331293": { "model_module": "@jupyter-widgets/controls", "model_module_version": "1.5.0", "model_name": "HTMLModel", "state": { "_dom_classes": [], "_model_module": "@jupyter-widgets/controls", "_model_module_version": "1.5.0", "_model_name": "HTMLModel", "_view_count": null, "_view_module": "@jupyter-widgets/controls", "_view_module_version": "1.5.0", "_view_name": "HTMLView", "description": "", "description_tooltip": null, "layout": "IPY_MODEL_e191f4a2517741689d2d11800da5cc45", "placeholder": "​", "style": "IPY_MODEL_f3436e0825204faca0a94f6795d75e89", "value": "Iteration: 100%" } }, "153550864ab84636af528e7eb8a959a3": { "model_module": "@jupyter-widgets/controls", "model_module_version": "1.5.0", "model_name": "DescriptionStyleModel", "state": { "_model_module": "@jupyter-widgets/controls", "_model_module_version": "1.5.0", "_model_name": "DescriptionStyleModel", "_view_count": null, "_view_module": "@jupyter-widgets/base", "_view_module_version": "1.2.0", "_view_name": "StyleView", "description_width": "" } }, "1b65756da31a43dd9ca4f49da6640e73": { "model_module": "@jupyter-widgets/base", "model_module_version": "1.2.0", "model_name": "LayoutModel", "state": { "_model_module": "@jupyter-widgets/base", "_model_module_version": "1.2.0", "_model_name": "LayoutModel", "_view_count": null, "_view_module": "@jupyter-widgets/base", "_view_module_version": "1.2.0", "_view_name": "LayoutView", "align_content": null, "align_items": null, "align_self": null, "border": null, "bottom": null, "display": null, "flex": null, "flex_flow": null, "grid_area": null, "grid_auto_columns": null, "grid_auto_flow": null, "grid_auto_rows": null, "grid_column": null, "grid_gap": null, "grid_row": null, "grid_template_areas": null, "grid_template_columns": null, "grid_template_rows": null, "height": null, "justify_content": null, "justify_items": null, "left": null, "margin": null, "max_height": null, "max_width": null, "min_height": null, "min_width": null, "object_fit": null, "object_position": null, "order": null, "overflow": null, "overflow_x": null, "overflow_y": null, "padding": null, "right": null, "top": null, "visibility": null, "width": null } }, "23f77655895c476d94cb02f152d0e4b7": { "model_module": "@jupyter-widgets/base", "model_module_version": "1.2.0", "model_name": "LayoutModel", "state": { "_model_module": "@jupyter-widgets/base", "_model_module_version": "1.2.0", "_model_name": "LayoutModel", "_view_count": null, "_view_module": "@jupyter-widgets/base", "_view_module_version": "1.2.0", "_view_name": "LayoutView", "align_content": null, "align_items": null, "align_self": null, "border": null, "bottom": null, "display": null, "flex": null, "flex_flow": null, "grid_area": null, "grid_auto_columns": null, "grid_auto_flow": null, "grid_auto_rows": null, "grid_column": null, "grid_gap": null, "grid_row": null, "grid_template_areas": null, "grid_template_columns": null, "grid_template_rows": null, "height": null, "justify_content": null, "justify_items": null, "left": null, "margin": null, "max_height": null, "max_width": null, "min_height": null, "min_width": null, "object_fit": null, "object_position": null, "order": null, "overflow": null, "overflow_x": null, "overflow_y": null, "padding": null, "right": null, "top": null, "visibility": null, "width": null } }, "31c15b576f134c2b8d6d14ca77dd2f88": { "model_module": "@jupyter-widgets/controls", "model_module_version": "1.5.0", "model_name": "DescriptionStyleModel", "state": { "_model_module": "@jupyter-widgets/controls", "_model_module_version": "1.5.0", "_model_name": "DescriptionStyleModel", "_view_count": null, "_view_module": "@jupyter-widgets/base", "_view_module_version": "1.2.0", "_view_name": "StyleView", "description_width": "" } }, "3a737d671bfb4aa6aee609a80ddede72": { "model_module": "@jupyter-widgets/controls", "model_module_version": "1.5.0", "model_name": "ProgressStyleModel", "state": { "_model_module": "@jupyter-widgets/controls", "_model_module_version": "1.5.0", "_model_name": "ProgressStyleModel", "_view_count": null, "_view_module": "@jupyter-widgets/base", "_view_module_version": "1.2.0", "_view_name": "StyleView", "bar_color": null, "description_width": "" } }, "4180378f634c4fd0a3c5013159cd6770": { "model_module": "@jupyter-widgets/controls", "model_module_version": "1.5.0", "model_name": "HBoxModel", "state": { "_dom_classes": [], "_model_module": "@jupyter-widgets/controls", "_model_module_version": "1.5.0", "_model_name": "HBoxModel", "_view_count": null, "_view_module": "@jupyter-widgets/controls", "_view_module_version": "1.5.0", "_view_name": "HBoxView", "box_style": "", "children": [ "IPY_MODEL_7c67c459c3c74f3eafe253e4865bed32", "IPY_MODEL_a31137fa736243509f27b08b0163630b", "IPY_MODEL_9bbfab5bf2354264879c63959d608daf" ], "layout": "IPY_MODEL_23f77655895c476d94cb02f152d0e4b7" } }, "7c67c459c3c74f3eafe253e4865bed32": { "model_module": "@jupyter-widgets/controls", "model_module_version": "1.5.0", "model_name": "HTMLModel", "state": { "_dom_classes": [], "_model_module": "@jupyter-widgets/controls", "_model_module_version": "1.5.0", "_model_name": "HTMLModel", "_view_count": null, "_view_module": "@jupyter-widgets/controls", "_view_module_version": "1.5.0", "_view_name": "HTMLView", "description": "", "description_tooltip": null, "layout": "IPY_MODEL_88b1e64d0a8149ebb81316c419c36a47", "placeholder": "​", "style": "IPY_MODEL_83ebfefc9345405bb0224c8df59e7d83", "value": "Epoch: 100%" } }, "7f933d6a976f4542bee79cf6bf1fe497": { "model_module": "@jupyter-widgets/controls", "model_module_version": "1.5.0", "model_name": "HBoxModel", "state": { "_dom_classes": [], "_model_module": "@jupyter-widgets/controls", "_model_module_version": "1.5.0", "_model_name": "HBoxModel", "_view_count": null, "_view_module": "@jupyter-widgets/controls", "_view_module_version": "1.5.0", "_view_name": "HBoxView", "box_style": "", "children": [ "IPY_MODEL_0ee3154796c44326989a4db3a3331293", "IPY_MODEL_df924c28c2214ab48c46999f8e332182", "IPY_MODEL_b39a63cc65224a8698cc2fa06af12ead" ], "layout": "IPY_MODEL_ed149912957740a393a0029497d4a144" } }, "83ebfefc9345405bb0224c8df59e7d83": { "model_module": "@jupyter-widgets/controls", "model_module_version": "1.5.0", "model_name": "DescriptionStyleModel", "state": { "_model_module": "@jupyter-widgets/controls", "_model_module_version": "1.5.0", "_model_name": "DescriptionStyleModel", "_view_count": null, "_view_module": "@jupyter-widgets/base", "_view_module_version": "1.2.0", "_view_name": "StyleView", "description_width": "" } }, "88b1e64d0a8149ebb81316c419c36a47": { "model_module": "@jupyter-widgets/base", "model_module_version": "1.2.0", "model_name": "LayoutModel", "state": { "_model_module": "@jupyter-widgets/base", "_model_module_version": "1.2.0", "_model_name": "LayoutModel", "_view_count": null, "_view_module": "@jupyter-widgets/base", "_view_module_version": "1.2.0", "_view_name": "LayoutView", "align_content": null, "align_items": null, "align_self": null, "border": null, "bottom": null, "display": null, "flex": null, "flex_flow": null, "grid_area": null, "grid_auto_columns": null, "grid_auto_flow": null, "grid_auto_rows": null, "grid_column": null, "grid_gap": null, "grid_row": null, "grid_template_areas": null, "grid_template_columns": null, "grid_template_rows": null, "height": null, "justify_content": null, "justify_items": null, "left": null, "margin": null, "max_height": null, "max_width": null, "min_height": null, "min_width": null, "object_fit": null, "object_position": null, "order": null, "overflow": null, "overflow_x": null, "overflow_y": null, "padding": null, "right": null, "top": null, "visibility": null, "width": null } }, "8f1c94ebd5a44c1ca9bcb9cfe8a85d56": { "model_module": "@jupyter-widgets/base", "model_module_version": "1.2.0", "model_name": "LayoutModel", "state": { "_model_module": "@jupyter-widgets/base", "_model_module_version": "1.2.0", "_model_name": "LayoutModel", "_view_count": null, "_view_module": "@jupyter-widgets/base", "_view_module_version": "1.2.0", "_view_name": "LayoutView", "align_content": null, "align_items": null, "align_self": null, "border": null, "bottom": null, "display": null, "flex": null, "flex_flow": null, "grid_area": null, "grid_auto_columns": null, "grid_auto_flow": null, "grid_auto_rows": null, "grid_column": null, "grid_gap": null, "grid_row": null, "grid_template_areas": null, "grid_template_columns": null, "grid_template_rows": null, "height": null, "justify_content": null, "justify_items": null, "left": null, "margin": null, "max_height": null, "max_width": null, "min_height": null, "min_width": null, "object_fit": null, "object_position": null, "order": null, "overflow": null, "overflow_x": null, "overflow_y": null, "padding": null, "right": null, "top": null, "visibility": null, "width": null } }, "9bbfab5bf2354264879c63959d608daf": { "model_module": "@jupyter-widgets/controls", "model_module_version": "1.5.0", "model_name": "HTMLModel", "state": { "_dom_classes": [], "_model_module": "@jupyter-widgets/controls", "_model_module_version": "1.5.0", "_model_name": "HTMLModel", "_view_count": null, "_view_module": "@jupyter-widgets/controls", "_view_module_version": "1.5.0", "_view_name": "HTMLView", "description": "", "description_tooltip": null, "layout": "IPY_MODEL_8f1c94ebd5a44c1ca9bcb9cfe8a85d56", "placeholder": "​", "style": "IPY_MODEL_31c15b576f134c2b8d6d14ca77dd2f88", "value": " 1/1 [35:19&lt;00:00, 2119.86s/it]" } }, "a31137fa736243509f27b08b0163630b": { "model_module": "@jupyter-widgets/controls", "model_module_version": "1.5.0", "model_name": "FloatProgressModel", "state": { "_dom_classes": [], "_model_module": "@jupyter-widgets/controls", "_model_module_version": "1.5.0", "_model_name": "FloatProgressModel", "_view_count": null, "_view_module": "@jupyter-widgets/controls", "_view_module_version": "1.5.0", "_view_name": "ProgressView", "bar_style": "success", "description": "", "description_tooltip": null, "layout": "IPY_MODEL_1b65756da31a43dd9ca4f49da6640e73", "max": 1, "min": 0, "orientation": "horizontal", "style": "IPY_MODEL_08ef465f79924616b29fe9828089d830", "value": 1 } }, "b39a63cc65224a8698cc2fa06af12ead": { "model_module": "@jupyter-widgets/controls", "model_module_version": "1.5.0", "model_name": "HTMLModel", "state": { "_dom_classes": [], "_model_module": "@jupyter-widgets/controls", "_model_module_version": "1.5.0", "_model_name": "HTMLModel", "_view_count": null, "_view_module": "@jupyter-widgets/controls", "_view_module_version": "1.5.0", "_view_name": "HTMLView", "description": "", "description_tooltip": null, "layout": "IPY_MODEL_d0202adacfa64dc6aeb0a3f4c33e28fc", "placeholder": "​", "style": "IPY_MODEL_153550864ab84636af528e7eb8a959a3", "value": " 4450/4450 [35:19&lt;00:00, 2.07it/s]" } }, "cc631484771647b989a24564792f29e4": { "model_module": "@jupyter-widgets/base", "model_module_version": "1.2.0", "model_name": "LayoutModel", "state": { "_model_module": "@jupyter-widgets/base", "_model_module_version": "1.2.0", "_model_name": "LayoutModel", "_view_count": null, "_view_module": "@jupyter-widgets/base", "_view_module_version": "1.2.0", "_view_name": "LayoutView", "align_content": null, "align_items": null, "align_self": null, "border": null, "bottom": null, "display": null, "flex": null, "flex_flow": null, "grid_area": null, "grid_auto_columns": null, "grid_auto_flow": null, "grid_auto_rows": null, "grid_column": null, "grid_gap": null, "grid_row": null, "grid_template_areas": null, "grid_template_columns": null, "grid_template_rows": null, "height": null, "justify_content": null, "justify_items": null, "left": null, "margin": null, "max_height": null, "max_width": null, "min_height": null, "min_width": null, "object_fit": null, "object_position": null, "order": null, "overflow": null, "overflow_x": null, "overflow_y": null, "padding": null, "right": null, "top": null, "visibility": null, "width": null } }, "d0202adacfa64dc6aeb0a3f4c33e28fc": { "model_module": "@jupyter-widgets/base", "model_module_version": "1.2.0", "model_name": "LayoutModel", "state": { "_model_module": "@jupyter-widgets/base", "_model_module_version": "1.2.0", "_model_name": "LayoutModel", "_view_count": null, "_view_module": "@jupyter-widgets/base", "_view_module_version": "1.2.0", "_view_name": "LayoutView", "align_content": null, "align_items": null, "align_self": null, "border": null, "bottom": null, "display": null, "flex": null, "flex_flow": null, "grid_area": null, "grid_auto_columns": null, "grid_auto_flow": null, "grid_auto_rows": null, "grid_column": null, "grid_gap": null, "grid_row": null, "grid_template_areas": null, "grid_template_columns": null, "grid_template_rows": null, "height": null, "justify_content": null, "justify_items": null, "left": null, "margin": null, "max_height": null, "max_width": null, "min_height": null, "min_width": null, "object_fit": null, "object_position": null, "order": null, "overflow": null, "overflow_x": null, "overflow_y": null, "padding": null, "right": null, "top": null, "visibility": null, "width": null } }, "df924c28c2214ab48c46999f8e332182": { "model_module": "@jupyter-widgets/controls", "model_module_version": "1.5.0", "model_name": "FloatProgressModel", "state": { "_dom_classes": [], "_model_module": "@jupyter-widgets/controls", "_model_module_version": "1.5.0", "_model_name": "FloatProgressModel", "_view_count": null, "_view_module": "@jupyter-widgets/controls", "_view_module_version": "1.5.0", "_view_name": "ProgressView", "bar_style": "success", "description": "", "description_tooltip": null, "layout": "IPY_MODEL_cc631484771647b989a24564792f29e4", "max": 4450, "min": 0, "orientation": "horizontal", "style": "IPY_MODEL_3a737d671bfb4aa6aee609a80ddede72", "value": 4450 } }, "e191f4a2517741689d2d11800da5cc45": { "model_module": "@jupyter-widgets/base", "model_module_version": "1.2.0", "model_name": "LayoutModel", "state": { "_model_module": "@jupyter-widgets/base", "_model_module_version": "1.2.0", "_model_name": "LayoutModel", "_view_count": null, "_view_module": "@jupyter-widgets/base", "_view_module_version": "1.2.0", "_view_name": "LayoutView", "align_content": null, "align_items": null, "align_self": null, "border": null, "bottom": null, "display": null, "flex": null, "flex_flow": null, "grid_area": null, "grid_auto_columns": null, "grid_auto_flow": null, "grid_auto_rows": null, "grid_column": null, "grid_gap": null, "grid_row": null, "grid_template_areas": null, "grid_template_columns": null, "grid_template_rows": null, "height": null, "justify_content": null, "justify_items": null, "left": null, "margin": null, "max_height": null, "max_width": null, "min_height": null, "min_width": null, "object_fit": null, "object_position": null, "order": null, "overflow": null, "overflow_x": null, "overflow_y": null, "padding": null, "right": null, "top": null, "visibility": null, "width": null } }, "ed149912957740a393a0029497d4a144": { "model_module": "@jupyter-widgets/base", "model_module_version": "1.2.0", "model_name": "LayoutModel", "state": { "_model_module": "@jupyter-widgets/base", "_model_module_version": "1.2.0", "_model_name": "LayoutModel", "_view_count": null, "_view_module": "@jupyter-widgets/base", "_view_module_version": "1.2.0", "_view_name": "LayoutView", "align_content": null, "align_items": null, "align_self": null, "border": null, "bottom": null, "display": null, "flex": null, "flex_flow": null, "grid_area": null, "grid_auto_columns": null, "grid_auto_flow": null, "grid_auto_rows": null, "grid_column": null, "grid_gap": null, "grid_row": null, "grid_template_areas": null, "grid_template_columns": null, "grid_template_rows": null, "height": null, "justify_content": null, "justify_items": null, "left": null, "margin": null, "max_height": null, "max_width": null, "min_height": null, "min_width": null, "object_fit": null, "object_position": null, "order": null, "overflow": null, "overflow_x": null, "overflow_y": null, "padding": null, "right": null, "top": null, "visibility": null, "width": null } }, "f3436e0825204faca0a94f6795d75e89": { "model_module": "@jupyter-widgets/controls", "model_module_version": "1.5.0", "model_name": "DescriptionStyleModel", "state": { "_model_module": "@jupyter-widgets/controls", "_model_module_version": "1.5.0", "_model_name": "DescriptionStyleModel", "_view_count": null, "_view_module": "@jupyter-widgets/base", "_view_module_version": "1.2.0", "_view_name": "StyleView", "description_width": "" } } } } }, "nbformat": 4, "nbformat_minor": 0 }