{ "cells": [ { "cell_type": "markdown", "metadata": {}, "source": [ "## Training metrics" ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "*Metrics* for training fastai models are simply functions that take `input` and `target` tensors, and return some metric of interest for training. You can write your own metrics by defining a function of that type, and passing it to [`Learner`](/basic_train.html#Learner) in the [`metrics`](/metrics.html#metrics) parameter, or use one of the following pre-defined functions." ] }, { "cell_type": "code", "execution_count": 1, "metadata": { "hide_input": true }, "outputs": [], "source": [ "from fastai.gen_doc.nbdoc import *\n", "from fastai.basics import *" ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "## Predefined metrics:" ] }, { "cell_type": "code", "execution_count": 2, "metadata": { "hide_input": true }, "outputs": [ { "data": { "text/markdown": [ "

accuracy[source][test]

\n", "\n", "> accuracy(**`input`**:`Tensor`, **`targs`**:`Tensor`) → `Rank0Tensor`\n", "\n", "
×

Tests found for accuracy:

To run tests please refer to this guide.

\n", "\n", "Computes accuracy with `targs` when `input` is bs * n_classes. " ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(accuracy)" ] }, { "cell_type": "code", "execution_count": null, "metadata": { "hide_input": true }, "outputs": [ { "data": { "text/markdown": [ "
Warning: This metric is intended for classification of objects belonging to a single class.
" ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "jekyll_warn(\"This metric is intended for classification of objects belonging to a single class.\")" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [ { "data": { "text/plain": [ "tensor(0.6000)" ] }, "execution_count": null, "metadata": {}, "output_type": "execute_result" } ], "source": [ "preds = tensor([0.4, 0.6], [0.3, 0.7], [0.2, 0.8], [0.6, 0.4], [0.9, 0.1]) # bs = 5, n = 2\n", "ys = tensor([1], [0], [1], [0], [1])\n", "accuracy(preds, ys)" ] }, { "cell_type": "code", "execution_count": 3, "metadata": { "hide_input": true }, "outputs": [ { "data": { "text/markdown": [ "

accuracy_thresh[source][test]

\n", "\n", "> accuracy_thresh(**`y_pred`**:`Tensor`, **`y_true`**:`Tensor`, **`thresh`**:`float`=***`0.5`***, **`sigmoid`**:`bool`=***`True`***) → `Rank0Tensor`\n", "\n", "
×

Tests found for accuracy_thresh:

  • pytest -sv tests/test_metrics.py::test_accuracy_thresh [source]

To run tests please refer to this guide.

\n", "\n", "Computes accuracy when `y_pred` and `y_true` are the same size. " ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(accuracy_thresh)" ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "Predictions are compared to `thresh` after `sigmoid` is maybe applied. Then we count the numbers that match the targets." ] }, { "cell_type": "code", "execution_count": null, "metadata": { "hide_input": true }, "outputs": [ { "data": { "text/markdown": [ "
Note: This function is intended for one-hot-encoded targets (often in a multiclassification problem).
" ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "jekyll_note(\"This function is intended for one-hot-encoded targets (often in a multiclassification problem).\")" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [ { "data": { "text/plain": [ "tensor(0.4000)" ] }, "execution_count": null, "metadata": {}, "output_type": "execute_result" } ], "source": [ "preds = tensor([0.4, 0.6], [0.3, 0.7], [0.2, 0.8], [0.6, 0.4], [0.9, 0.1])\n", "ys = tensor([0, 1], [1, 0], [0, 1], [1, 0], [0, 1]) \n", "accuracy_thresh(preds, ys, thresh=0.65, sigmoid=False)" ] }, { "cell_type": "code", "execution_count": 4, "metadata": { "hide_input": true }, "outputs": [ { "data": { "text/markdown": [ "

top_k_accuracy[source][test]

\n", "\n", "> top_k_accuracy(**`input`**:`Tensor`, **`targs`**:`Tensor`, **`k`**:`int`=***`5`***) → `Rank0Tensor`\n", "\n", "
×

Tests found for top_k_accuracy:

  • pytest -sv tests/test_metrics.py::test_top_k_accuracy [source]

To run tests please refer to this guide.

\n", "\n", "Computes the Top-k accuracy (target is in the top k predictions). " ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(top_k_accuracy)" ] }, { "cell_type": "code", "execution_count": 5, "metadata": { "hide_input": true }, "outputs": [ { "data": { "text/markdown": [ "

dice[source][test]

\n", "\n", "> dice(**`input`**:`Tensor`, **`targs`**:`Tensor`, **`iou`**:`bool`=***`False`***, **`eps`**:`float`=***`1e-08`***) → `Rank0Tensor`\n", "\n", "
×

Tests found for dice:

  • pytest -sv tests/test_metrics.py::test_dice [source]
  • pytest -sv tests/test_metrics.py::test_dice_iou [source]

To run tests please refer to this guide.

\n", "\n", "Dice coefficient metric for binary target. If iou=True, returns iou metric, classic for segmentation problems. " ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(dice)" ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "$$dice = \\frac{2(TP)}{2(TP) + FP + FN}$$\n", "\n", "where TP, FP and FN are the number of true positives, false positives and false negatives." ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [ { "data": { "text/plain": [ "tensor(0.6667)" ] }, "execution_count": null, "metadata": {}, "output_type": "execute_result" } ], "source": [ "preds = tensor([0.4, 0.6], [0.3, 0.7], [0.2, 0.8], [0.6, 0.4], [0.9, 0.1])\n", "ys = tensor([1], [0], [1], [0], [1])\n", "dice(preds, ys) # TP = 2, FP = 1, FN = 1" ] }, { "cell_type": "code", "execution_count": 6, "metadata": { "hide_input": true }, "outputs": [ { "data": { "text/markdown": [ "

error_rate[source][test]

\n", "\n", "> error_rate(**`input`**:`Tensor`, **`targs`**:`Tensor`) → `Rank0Tensor`\n", "\n", "
×

Tests found for error_rate:

  • pytest -sv tests/test_metrics.py::test_error_rate [source]
  • pytest -sv tests/test_vision_train.py::test_error_rate [source]

To run tests please refer to this guide.

\n", "\n", "1 - [`accuracy`](/metrics.html#accuracy) " ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(error_rate)" ] }, { "cell_type": "code", "execution_count": 7, "metadata": { "hide_input": true }, "outputs": [ { "data": { "text/markdown": [ "

mean_squared_error[source][test]

\n", "\n", "> mean_squared_error(**`pred`**:`Tensor`, **`targ`**:`Tensor`) → `Rank0Tensor`\n", "\n", "
×

Tests found for mean_squared_error:

  • pytest -sv tests/test_metrics.py::test_mse [source]

To run tests please refer to this guide.

\n", "\n", "Mean squared error between `pred` and `targ`. " ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(mean_squared_error)" ] }, { "cell_type": "code", "execution_count": 8, "metadata": { "hide_input": true }, "outputs": [ { "data": { "text/markdown": [ "

mean_absolute_error[source][test]

\n", "\n", "> mean_absolute_error(**`pred`**:`Tensor`, **`targ`**:`Tensor`) → `Rank0Tensor`\n", "\n", "
×

Tests found for mean_absolute_error:

  • pytest -sv tests/test_metrics.py::test_mae [source]

To run tests please refer to this guide.

\n", "\n", "Mean absolute error between `pred` and `targ`. " ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(mean_absolute_error)" ] }, { "cell_type": "code", "execution_count": 9, "metadata": { "hide_input": true }, "outputs": [ { "data": { "text/markdown": [ "

mean_squared_logarithmic_error[source][test]

\n", "\n", "> mean_squared_logarithmic_error(**`pred`**:`Tensor`, **`targ`**:`Tensor`) → `Rank0Tensor`\n", "\n", "
×

Tests found for mean_squared_logarithmic_error:

  • pytest -sv tests/test_metrics.py::test_msle [source]

To run tests please refer to this guide.

\n", "\n", "Mean squared logarithmic error between `pred` and `targ`. " ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(mean_squared_logarithmic_error)" ] }, { "cell_type": "code", "execution_count": 10, "metadata": { "hide_input": true }, "outputs": [ { "data": { "text/markdown": [ "

exp_rmspe[source][test]

\n", "\n", "> exp_rmspe(**`pred`**:`Tensor`, **`targ`**:`Tensor`) → `Rank0Tensor`\n", "\n", "
×

Tests found for exp_rmspe:

  • pytest -sv tests/test_metrics.py::test_exp_rmspe [source]
  • pytest -sv tests/test_metrics.py::test_exp_rmspe_num_of_ele [source]

To run tests please refer to this guide.

\n", "\n", "Exp RMSE between `pred` and `targ`. " ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(exp_rmspe)" ] }, { "cell_type": "code", "execution_count": 11, "metadata": { "hide_input": true }, "outputs": [ { "data": { "text/markdown": [ "

root_mean_squared_error[source][test]

\n", "\n", "> root_mean_squared_error(**`pred`**:`Tensor`, **`targ`**:`Tensor`) → `Rank0Tensor`\n", "\n", "
×

Tests found for root_mean_squared_error:

  • pytest -sv tests/test_metrics.py::test_rmse [source]

To run tests please refer to this guide.

\n", "\n", "Root mean squared error between `pred` and `targ`. " ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(root_mean_squared_error)" ] }, { "cell_type": "code", "execution_count": 12, "metadata": { "hide_input": true }, "outputs": [ { "data": { "text/markdown": [ "

fbeta[source][test]

\n", "\n", "> fbeta(**`y_pred`**:`Tensor`, **`y_true`**:`Tensor`, **`thresh`**:`float`=***`0.2`***, **`beta`**:`float`=***`2`***, **`eps`**:`float`=***`1e-09`***, **`sigmoid`**:`bool`=***`True`***) → `Rank0Tensor`\n", "\n", "
×

Tests found for fbeta:

  • pytest -sv tests/test_metrics.py::test_fbeta [source]

To run tests please refer to this guide.

\n", "\n", "Computes the f_beta between `preds` and `targets` " ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(fbeta)" ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "`beta` determines the value of the fbeta applied, `eps` is there for numeric stability. If `sigmoid=True`, a sigmoid is applied to the predictions before comparing them to `thresh` then to the targets. See the [F1 score wikipedia page](https://en.wikipedia.org/wiki/F1_score) for details on the fbeta score." ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "$${F_\\beta} = (1+\\beta^2)\\frac{precision \\cdot recall}{(\\beta^2 \\cdot precision) + recall}$$" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [ { "data": { "text/plain": [ "tensor(0.6667)" ] }, "execution_count": null, "metadata": {}, "output_type": "execute_result" } ], "source": [ "preds = tensor([0.6, 0.8, 0.2, 0.4, 0.9]).view(1, 5) # TP =2, FP = 1, FN = 1\n", "ys = tensor([1, 0, 0, 1, 1]).view(1, 5)\n", "fbeta(preds, ys, thresh=0.5, sigmoid=False)" ] }, { "cell_type": "code", "execution_count": null, "metadata": { "hide_input": true }, "outputs": [ { "data": { "text/markdown": [ "
Note: This function is intended for one-hot-encoded targets (often in a multiclassification problem).
" ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "jekyll_note(\"This function is intended for one-hot-encoded targets (often in a multiclassification problem).\")" ] }, { "cell_type": "code", "execution_count": 13, "metadata": { "hide_input": true }, "outputs": [ { "data": { "text/markdown": [ "

explained_variance[source][test]

\n", "\n", "> explained_variance(**`pred`**:`Tensor`, **`targ`**:`Tensor`) → `Rank0Tensor`\n", "\n", "
×

Tests found for explained_variance:

  • pytest -sv tests/test_metrics.py::test_explained_variance [source]

To run tests please refer to this guide.

\n", "\n", "Explained variance between `pred` and `targ`. " ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(explained_variance)" ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "$$ Explained \\ Variance = 1 - \\frac{Var( targ - pred )}{Var( targ )}$$" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [ { "data": { "text/plain": [ "tensor(0.9374)" ] }, "execution_count": null, "metadata": {}, "output_type": "execute_result" } ], "source": [ "preds = tensor([0.10, .20, .30, .40, .50])\n", "ys = tensor([0.12, .17, .25, .44, .56]) # predictions are close to the truth\n", "explained_variance(preds, ys)" ] }, { "cell_type": "code", "execution_count": 14, "metadata": { "hide_input": true }, "outputs": [ { "data": { "text/markdown": [ "

r2_score[source][test]

\n", "\n", "> r2_score(**`pred`**:`Tensor`, **`targ`**:`Tensor`) → `Rank0Tensor`\n", "\n", "
×

Tests found for r2_score:

  • pytest -sv tests/test_metrics.py::test_r2_score [source]

To run tests please refer to this guide.

\n", "\n", "R2 score (coefficient of determination) between `pred` and `targ`. " ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(r2_score)" ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "$$ {R^2} = 1 - \\frac{\\sum( targ - pred )^2}{\\sum( targ - \\overline{targ})^2}$$\n", "where $\\overline{targ}$ is the mean of the targ tensor." ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [ { "data": { "text/plain": [ "tensor(0.9351)" ] }, "execution_count": null, "metadata": {}, "output_type": "execute_result" } ], "source": [ "r2_score(preds, ys)" ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "The following metrics are classes, don't forget to instantiate them when you pass them to a [`Learner`](/basic_train.html#Learner)." ] }, { "cell_type": "code", "execution_count": 15, "metadata": { "hide_input": true }, "outputs": [ { "data": { "text/markdown": [ "

class RMSE[source][test]

\n", "\n", "> RMSE() :: [`RegMetrics`](/metrics.html#RegMetrics)\n", "\n", "
×

No tests found for RMSE. To contribute a test please refer to this guide and this discussion.

\n", "\n", "Computes the root mean squared error. " ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(RMSE, title_level=3)" ] }, { "cell_type": "code", "execution_count": 16, "metadata": { "hide_input": true }, "outputs": [ { "data": { "text/markdown": [ "

class ExpRMSPE[source][test]

\n", "\n", "> ExpRMSPE() :: [`RegMetrics`](/metrics.html#RegMetrics)\n", "\n", "
×

No tests found for ExpRMSPE. To contribute a test please refer to this guide and this discussion.

\n", "\n", "Computes the exponential of the root mean square error. " ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(ExpRMSPE, title_level=3)" ] }, { "cell_type": "code", "execution_count": 17, "metadata": { "hide_input": true }, "outputs": [ { "data": { "text/markdown": [ "

class Precision[source][test]

\n", "\n", "> Precision(**`average`**:`Optional`\\[`str`\\]=***`'binary'`***, **`pos_label`**:`int`=***`1`***, **`eps`**:`float`=***`1e-09`***) :: [`CMScores`](/metrics.html#CMScores)\n", "\n", "
×

No tests found for Precision. To contribute a test please refer to this guide and this discussion.

\n", "\n", "Computes the Precision. " ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(Precision, title_level=3)" ] }, { "cell_type": "code", "execution_count": 18, "metadata": { "hide_input": true }, "outputs": [ { "data": { "text/markdown": [ "

class Recall[source][test]

\n", "\n", "> Recall(**`average`**:`Optional`\\[`str`\\]=***`'binary'`***, **`pos_label`**:`int`=***`1`***, **`eps`**:`float`=***`1e-09`***) :: [`CMScores`](/metrics.html#CMScores)\n", "\n", "
×

No tests found for Recall. To contribute a test please refer to this guide and this discussion.

\n", "\n", "Computes the Recall. " ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(Recall, title_level=3)" ] }, { "cell_type": "code", "execution_count": 19, "metadata": { "hide_input": true }, "outputs": [ { "data": { "text/markdown": [ "

class FBeta[source][test]

\n", "\n", "> FBeta(**`average`**:`Optional`\\[`str`\\]=***`'binary'`***, **`pos_label`**:`int`=***`1`***, **`eps`**:`float`=***`1e-09`***, **`beta`**:`float`=***`2`***) :: [`CMScores`](/metrics.html#CMScores)\n", "\n", "
×

No tests found for FBeta. To contribute a test please refer to this guide and this discussion.

\n", "\n", "Computes the F`beta` score. " ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(FBeta, title_level=3)" ] }, { "cell_type": "code", "execution_count": 20, "metadata": { "hide_input": true }, "outputs": [ { "data": { "text/markdown": [ "

class R2Score[source][test]

\n", "\n", "> R2Score() :: [`RegMetrics`](/metrics.html#RegMetrics)\n", "\n", "
×

No tests found for R2Score. To contribute a test please refer to this guide and this discussion.

\n", "\n", "Computes the R2 score (coefficient of determination). " ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(R2Score, title_level=3)" ] }, { "cell_type": "code", "execution_count": 21, "metadata": { "hide_input": true }, "outputs": [ { "data": { "text/markdown": [ "

class ExplainedVariance[source][test]

\n", "\n", "> ExplainedVariance() :: [`RegMetrics`](/metrics.html#RegMetrics)\n", "\n", "
×

No tests found for ExplainedVariance. To contribute a test please refer to this guide and this discussion.

\n", "\n", "Computes the explained variance. " ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(ExplainedVariance, title_level=3)" ] }, { "cell_type": "code", "execution_count": 22, "metadata": { "hide_input": true }, "outputs": [ { "data": { "text/markdown": [ "

class MatthewsCorreff[source][test]

\n", "\n", "> MatthewsCorreff() :: [`ConfusionMatrix`](/metrics.html#ConfusionMatrix)\n", "\n", "
×

No tests found for MatthewsCorreff. To contribute a test please refer to this guide and this discussion.

\n", "\n", "Computes the Matthews correlation coefficient. " ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(MatthewsCorreff, title_level=3)" ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "Ref.: https://github.com/scikit-learn/scikit-learn/blob/bac89c2/sklearn/metrics/classification.py" ] }, { "cell_type": "code", "execution_count": 23, "metadata": { "hide_input": true }, "outputs": [ { "data": { "text/markdown": [ "

class KappaScore[source][test]

\n", "\n", "> KappaScore(**`weights`**:`Optional`\\[`str`\\]=***`None`***) :: [`ConfusionMatrix`](/metrics.html#ConfusionMatrix)\n", "\n", "
×

No tests found for KappaScore. To contribute a test please refer to this guide and this discussion.

\n", "\n", "Computes the rate of agreement (Cohens Kappa). " ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(KappaScore, title_level=3)" ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "Ref.: https://github.com/scikit-learn/scikit-learn/blob/bac89c2/sklearn/metrics/classification.py \n", "\n", "[`KappaScore`](/metrics.html#KappaScore) supports linear and quadratic weights on the off-diagonal cells in the [`ConfusionMatrix`](/metrics.html#ConfusionMatrix), in addition to the default unweighted calculation treating all misclassifications as equally weighted. Leaving [`KappaScore`](/metrics.html#KappaScore)'s `weights` attribute as `None` returns the unweighted Kappa score. Updating `weights` to \"linear\" means off-diagonal ConfusionMatrix elements are weighted in linear proportion to their distance from the diagonal; \"quadratic\" means weights are squared proportional to their distance from the diagonal.\n", "Specify linear or quadratic weights, if using, by first creating an instance of the metric and then updating the `weights` attribute, similar to as follows: \n", "```\n", "kappa = KappaScore()\n", "kappa.weights = \"quadratic\"\n", "learn = cnn_learner(data, model, metrics=[error_rate, kappa])\n", "``` " ] }, { "cell_type": "code", "execution_count": 24, "metadata": { "hide_input": true }, "outputs": [ { "data": { "text/markdown": [ "

class ConfusionMatrix[source][test]

\n", "\n", "> ConfusionMatrix() :: [`Callback`](/callback.html#Callback)\n", "\n", "
×

No tests found for ConfusionMatrix. To contribute a test please refer to this guide and this discussion.

\n", "\n", "Computes the confusion matrix. " ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(ConfusionMatrix, title_level=3)" ] }, { "cell_type": "code", "execution_count": 25, "metadata": { "hide_input": true }, "outputs": [ { "data": { "text/markdown": [ "

class MultiLabelFbeta[source][test]

\n", "\n", "> MultiLabelFbeta(**`beta`**=***`2`***, **`eps`**=***`1e-15`***, **`thresh`**=***`0.3`***, **`sigmoid`**=***`True`***, **`average`**=***`'micro'`***) :: [`Callback`](/callback.html#Callback)\n", "\n", "
×

No tests found for MultiLabelFbeta. To contribute a test please refer to this guide and this discussion.

\n", "\n", "Computes the fbeta score for multilabel classification " ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(MultiLabelFbeta, title_level=3)" ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "[`MultiLabelFbeta`](/metrics.html#MultiLabelFbeta) implements mutlilabel classification fbeta score similar to [scikit-learn's](https://scikit-learn.org/stable/modules/generated/sklearn.metrics.f1_score.html) as a [`LearnerCallback`](/basic_train.html#LearnerCallback). Average options: [\"micro\", \"macro\", \"weighted\", \"none\"]. Intended to use with one-hot encoded targets with 1s and 0s." ] }, { "cell_type": "code", "execution_count": 26, "metadata": {}, "outputs": [ { "data": { "text/markdown": [ "

auc_roc_score[source][test]

\n", "\n", "> auc_roc_score(**`input`**:`Tensor`, **`targ`**:`Tensor`)\n", "\n", "
×

No tests found for auc_roc_score. To contribute a test please refer to this guide and this discussion.

\n", "\n", "Computes the area under the receiver operator characteristic (ROC) curve using the trapezoid method. Restricted binary classification tasks. " ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(auc_roc_score, title_level=3)" ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "[`auc_roc_score`](/metrics.html#auc_roc_score) computes the AUC score for the ROC curve similarly to [scikit-learn](https://scikit-learn.org/stable/modules/generated/sklearn.metrics.roc_auc_score.html) using the trapezoid method, effectively summarizing the curve information in a single number. See [Wikipedia's page](https://en.wikipedia.org/wiki/Receiver_operating_characteristic#Area_under_the_curve) for more information on this." ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [ { "data": { "text/markdown": [ "
Note: Instead of passing this method to the learner's metrics directly, make use of the AUROC() class.
" ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "jekyll_note(\"Instead of passing this method to the learner's metrics directly, make use of the AUROC() class.\")" ] }, { "cell_type": "code", "execution_count": 27, "metadata": {}, "outputs": [ { "data": { "text/markdown": [ "

roc_curve[source][test]

\n", "\n", "> roc_curve(**`input`**:`Tensor`, **`targ`**:`Tensor`)\n", "\n", "
×

No tests found for roc_curve. To contribute a test please refer to this guide and this discussion.

\n", "\n", "Computes the receiver operator characteristic (ROC) curve by determining the true positive ratio (TPR) and false positive ratio (FPR) for various classification thresholds. Restricted binary classification tasks. " ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(roc_curve, title_level=3)" ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "[`roc_curve`](/metrics.html#roc_curve) generates the ROC curve similarly to [scikit-learn](https://scikit-learn.org/stable/modules/generated/sklearn.metrics.roc_curve.html). See [Wikipedia's page](https://en.wikipedia.org/wiki/Receiver_operating_characteristic) for more information on the ROC curve." ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [ { "data": { "text/markdown": [ "
Note: Instead of passing this method to the learner's metrics directly, make use of the AUROC() class.
" ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "jekyll_note(\"Instead of passing this method to the learner's metrics directly, make use of the AUROC() class.\")" ] }, { "cell_type": "code", "execution_count": 28, "metadata": {}, "outputs": [ { "data": { "text/markdown": [ "

class AUROC[source][test]

\n", "\n", "> AUROC() :: [`Callback`](/callback.html#Callback)\n", "\n", "
×

No tests found for AUROC. To contribute a test please refer to this guide and this discussion.

\n", "\n", "Computes the area under the curve (AUC) score based on the receiver operator characteristic (ROC) curve. Restricted to binary classification tasks. " ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(AUROC, title_level=3)" ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "[`AUROC`](/metrics.html#AUROC) creates a [`Callback`](/callback.html#Callback) for computing the AUC score for the ROC curve with [`auc_roc_score`](/metrics.html#auc_roc_score) at the end of each epoch, given that averaging over batches is incorrect in case of the AUROC. See [Wikipedia's page](https://en.wikipedia.org/wiki/Receiver_operating_characteristic#Area_under_the_curve) for more information on the AUROC." ] }, { "cell_type": "markdown", "metadata": { "hide_input": true }, "source": [ "## Creating your own metric" ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "Creating a new metric can be as simple as creating a new function. If your metric is an average over the total number of elements in your dataset, just write the function that will compute it on a batch (taking `pred` and `targ` as arguments). It will then be automatically averaged over the batches (taking their different sizes into account).\n", "\n", "Sometimes metrics aren't simple averages however. If we take the example of precision for instance, we have to divide the number of true positives by the number of predictions we made for that class. This isn't an average over the number of elements we have in the dataset, we only consider those where we made a positive prediction for a specific thing. Computing the precision for each batch, then averaging them will yield to a result that may be close to the real value, but won't be it exactly (and it really depends on how you deal with special case of 0 positive predictions).\n", "\n", "This why in fastai, every metric is implemented as a callback. If you pass a regular function, the library transforms it to a proper callback called `AverageCallback`. The callback metrics are only called during the validation phase, and only for the following events: \n", "- on_epoch_begin (for initialization)\n", "- on_batch_begin (if we need to have a look at the input/target and maybe modify them)\n", "- on_batch_end (to analyze the last results and update our computation)\n", "- on_epoch_end(to wrap up the final result that should be added to `last_metrics`)\n", "\n", "As an example, the following code is the exact implementation of the [`AverageMetric`](/callback.html#AverageMetric) callback that transforms a function like [`accuracy`](/metrics.html#accuracy) into a metric callback." ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "class AverageMetric(Callback):\n", " \"Wrap a `func` in a callback for metrics computation.\"\n", " def __init__(self, func):\n", " # If it's a partial, use func.func\n", " name = getattr(func,'func',func).__name__\n", " self.func, self.name = func, name\n", "\n", " def on_epoch_begin(self, **kwargs):\n", " \"Set the inner value to 0.\"\n", " self.val, self.count = 0.,0\n", "\n", " def on_batch_end(self, last_output, last_target, **kwargs):\n", " \"Update metric computation with `last_output` and `last_target`.\"\n", " if not is_listy(last_target): last_target=[last_target]\n", " self.count += last_target[0].size(0)\n", " val = self.func(last_output, *last_target)\n", " self.val += last_target[0].size(0) * val.detach().cpu()\n", "\n", " def on_epoch_end(self, last_metrics, **kwargs):\n", " \"Set the final result in `last_metrics`.\"\n", " return add_metrics(last_metrics, self.val/self.count)" ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "Here [`add_metrics`](/torch_core.html#add_metrics) is a convenience function that will return the proper dictionary for us:\n", "```python\n", "{'last_metrics': last_metrics + [self.val/self.count]}\n", "```" ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "And here is another example that properly computes the precision for a given class." ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "class Precision(Callback):\n", " \n", " def on_epoch_begin(self, **kwargs):\n", " self.correct, self.total = 0, 0\n", " \n", " def on_batch_end(self, last_output, last_target, **kwargs):\n", " preds = last_output.argmax(1)\n", " self.correct += ((preds==0) * (last_target==0)).float().sum()\n", " self.total += (preds==0).float().sum()\n", " \n", " def on_epoch_end(self, last_metrics, **kwargs):\n", " return add_metrics(last_metrics, self.correct/self.total)" ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "The following custom callback class example measures peak RAM usage during each epoch: " ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "import tracemalloc\n", "class TraceMallocMetric(Callback):\n", " def __init__(self):\n", " super().__init__()\n", " self.name = \"peak RAM\"\n", "\n", " def on_epoch_begin(self, **kwargs):\n", " tracemalloc.start()\n", " \n", " def on_epoch_end(self, last_metrics, **kwargs):\n", " current, peak = tracemalloc.get_traced_memory()\n", " tracemalloc.stop()\n", " return add_metrics(last_metrics, torch.tensor(peak))" ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "To deploy it, you need to pass an instance of this custom metric in the [`metrics`](/metrics.html#metrics) argument:\n", "\n", "```python\n", "learn = cnn_learner(data, model, metrics=[accuracy, TraceMallocMetric()])\n", "learn.fit_one_cycle(3, max_lr=1e-2)\n", "```\n", "And then the output changes to:\n", "```\n", "Total time: 00:54\n", "epoch\ttrain_loss\tvalid_loss\taccuracy\tpeak RAM\n", " 1\t0.333352\t0.084342\t0.973800\t2395541.000000\n", " 2\t0.096196\t0.038386\t0.988300\t2342145.000000\n", " 3\t0.048722\t0.029234\t0.990200\t2342680.000000\n", "```\n" ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "As mentioner earlier, using the [`metrics`](/metrics.html#metrics) argument with a custom metrics class is limited in the number of phases of the callback system it can access, it can only return one numerical value and as you can see its output is hardcoded to have 6 points of precision in the output, even if the number is an int.\n", "\n", "To overcome these limitations callback classes should be used instead.\n", "\n", "For example, the following class: \n", "* uses phases not available for the metric classes \n", "* it reports 3 columns, instead of just one\n", "* its column report ints, instead of floats" ] }, { "cell_type": "code", "execution_count": null, "metadata": {}, "outputs": [], "source": [ "import tracemalloc\n", "class TraceMallocMultiColMetric(LearnerCallback):\n", " _order=-20 # Needs to run before the recorder\n", " def __init__(self, learn):\n", " super().__init__(learn)\n", " self.train_max = 0\n", "\n", " def on_train_begin(self, **kwargs):\n", " self.learn.recorder.add_metric_names(['used', 'max_used', 'peak'])\n", " \n", " def on_batch_end(self, train, **kwargs):\n", " # track max memory usage during the train phase\n", " if train:\n", " current, peak = tracemalloc.get_traced_memory()\n", " self.train_max = max(self.train_max, current)\n", " \n", " def on_epoch_begin(self, **kwargs):\n", " tracemalloc.start()\n", "\n", " def on_epoch_end(self, last_metrics, **kwargs):\n", " current, peak = tracemalloc.get_traced_memory()\n", " tracemalloc.stop()\n", " return add_metrics(last_metrics, [current, self.train_max, peak])" ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "Note, that it subclasses [`LearnerCallback`](/basic_train.html#LearnerCallback) and not [`Callback`](/callback.html#Callback), since the former provides extra features not available in the latter.\n", "\n", "Also `_order=-20` is crucial - without it the custom columns will not be added - it tells the callback system to run this callback before the recorder system. \n", "\n", "To deploy it, you need to pass the name of the class (not an instance!) of the class in the `callback_fns` argument. This is because the `learn` object doesn't exist yet, and it's required to instantiate `TraceMallocMultiColMetric`. The system will do it for us automatically as soon as the learn object has been created.\n", "\n", "```python\n", "learn = cnn_learner(data, model, metrics=[accuracy], callback_fns=TraceMallocMultiColMetric)\n", "learn.fit_one_cycle(3, max_lr=1e-2)\n", "```\n", "And then the output changes to:\n", "```\n", "Total time: 00:53\n", "epoch\ttrain_loss valid_loss accuracy\t used\tmax_used peak\n", " 1\t0.321233\t0.068252\t0.978600\t156504\t2408404\t 2419891 \n", " 2\t0.093551\t0.032776\t0.988500\t 79343\t2408404\t 2348085\n", " 3\t0.047178\t0.025307\t0.992100\t 79568\t2408404\t 2342754\n", "```\n", "\n", "Another way to do the same is by using `learn.callbacks.append`, and this time we need to instantiate `TraceMallocMultiColMetric` with `learn` object which we now have, as it is called after the latter was created:\n", "\n", "```python\n", "learn = cnn_learner(data, model, metrics=[accuracy])\n", "learn.callbacks.append(TraceMallocMultiColMetric(learn))\n", "learn.fit_one_cycle(3, max_lr=1e-2)\n", "```\n", "\n", "Configuring the custom metrics in the `learn` object sets them to run in all future [`fit`](/basic_train.html#fit)-family calls. However, if you'd like to configure it for just one call, you can configure it directly inside [`fit`](/basic_train.html#fit) or [`fit_one_cycle`](/train.html#fit_one_cycle):\n", "\n", "```python\n", "learn = cnn_learner(data, model, metrics=[accuracy])\n", "learn.fit_one_cycle(3, max_lr=1e-2, callbacks=TraceMallocMultiColMetric(learn))\n", "```\n", "\n", "And to stress the differences: \n", "* the `callback_fns` argument expects a classname or a list of those\n", "* the [`callbacks`](/callbacks.html#callbacks) argument expects an instance of a class or a list of those\n", "* `learn.callbacks.append` expects a single instance of a class\n", "\n", "For more examples, look inside fastai codebase and its test suite, search for classes that subclass either [`Callback`](/callback.html#Callback), [`LearnerCallback`](/basic_train.html#LearnerCallback) and subclasses of those two.\n", "\n", "Finally, while the above examples all add to the metrics, it's not a requirement. A callback can do anything it wants and it is not required to add its outcomes to the metrics printout. " ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "## Undocumented Methods - Methods moved below this line will intentionally be hidden" ] }, { "cell_type": "code", "execution_count": 29, "metadata": {}, "outputs": [ { "data": { "text/markdown": [ "

on_batch_end[source][test]

\n", "\n", "> on_batch_end(**`last_output`**:`Tensor`, **`last_target`**:`Tensor`, **\\*\\*`kwargs`**)\n", "\n", "
×

No tests found for on_batch_end. To contribute a test please refer to this guide and this discussion.

\n", "\n", "Called at the end of the batch. " ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(FBeta.on_batch_end)" ] }, { "cell_type": "code", "execution_count": 30, "metadata": {}, "outputs": [ { "data": { "text/markdown": [ "

on_epoch_begin[source][test]

\n", "\n", "> on_epoch_begin(**\\*\\*`kwargs`**)\n", "\n", "
×

No tests found for on_epoch_begin. To contribute a test please refer to this guide and this discussion.

\n", "\n", "At the beginning of each epoch. " ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(FBeta.on_epoch_begin)" ] }, { "cell_type": "code", "execution_count": 31, "metadata": {}, "outputs": [ { "data": { "text/markdown": [ "

on_epoch_end[source][test]

\n", "\n", "> on_epoch_end(**`last_metrics`**, **\\*\\*`kwargs`**)\n", "\n", "
×

No tests found for on_epoch_end. To contribute a test please refer to this guide and this discussion.

\n", "\n", "Called at the end of an epoch. " ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(FBeta.on_epoch_end)" ] }, { "cell_type": "code", "execution_count": 32, "metadata": {}, "outputs": [ { "data": { "text/markdown": [ "

mean_absolute_error[source][test]

\n", "\n", "> mean_absolute_error(**`pred`**:`Tensor`, **`targ`**:`Tensor`) → `Rank0Tensor`\n", "\n", "
×

Tests found for mean_absolute_error:

  • pytest -sv tests/test_metrics.py::test_mae [source]

To run tests please refer to this guide.

\n", "\n", "Mean absolute error between `pred` and `targ`. " ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(mae)" ] }, { "cell_type": "code", "execution_count": 33, "metadata": {}, "outputs": [ { "data": { "text/markdown": [ "

mean_squared_logarithmic_error[source][test]

\n", "\n", "> mean_squared_logarithmic_error(**`pred`**:`Tensor`, **`targ`**:`Tensor`) → `Rank0Tensor`\n", "\n", "
×

Tests found for mean_squared_logarithmic_error:

  • pytest -sv tests/test_metrics.py::test_msle [source]

To run tests please refer to this guide.

\n", "\n", "Mean squared logarithmic error between `pred` and `targ`. " ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(msle)" ] }, { "cell_type": "code", "execution_count": 34, "metadata": {}, "outputs": [ { "data": { "text/markdown": [ "

mean_squared_error[source][test]

\n", "\n", "> mean_squared_error(**`pred`**:`Tensor`, **`targ`**:`Tensor`) → `Rank0Tensor`\n", "\n", "
×

Tests found for mean_squared_error:

  • pytest -sv tests/test_metrics.py::test_mse [source]

To run tests please refer to this guide.

\n", "\n", "Mean squared error between `pred` and `targ`. " ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(mse)" ] }, { "cell_type": "code", "execution_count": 35, "metadata": {}, "outputs": [ { "data": { "text/markdown": [ "

root_mean_squared_error[source][test]

\n", "\n", "> root_mean_squared_error(**`pred`**:`Tensor`, **`targ`**:`Tensor`) → `Rank0Tensor`\n", "\n", "
×

Tests found for root_mean_squared_error:

  • pytest -sv tests/test_metrics.py::test_rmse [source]

To run tests please refer to this guide.

\n", "\n", "Root mean squared error between `pred` and `targ`. " ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(rmse)" ] }, { "cell_type": "code", "execution_count": 36, "metadata": {}, "outputs": [ { "data": { "text/markdown": [ "

on_epoch_end[source][test]

\n", "\n", "> on_epoch_end(**`last_metrics`**, **\\*\\*`kwargs`**)\n", "\n", "
×

No tests found for on_epoch_end. To contribute a test please refer to this guide and this discussion.

\n", "\n", "Called at the end of an epoch. " ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(Precision.on_epoch_end)" ] }, { "cell_type": "code", "execution_count": 37, "metadata": {}, "outputs": [ { "data": { "text/markdown": [ "

on_train_end[source][test]

\n", "\n", "> on_train_end(**\\*\\*`kwargs`**)\n", "\n", "
×

No tests found for on_train_end. To contribute a test please refer to this guide and this discussion.

\n", "\n", "Useful for cleaning up things and saving files/models. " ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(FBeta.on_train_end)" ] }, { "cell_type": "code", "execution_count": 38, "metadata": {}, "outputs": [ { "data": { "text/markdown": [ "

on_epoch_end[source][test]

\n", "\n", "> on_epoch_end(**`last_metrics`**, **\\*\\*`kwargs`**)\n", "\n", "
×

No tests found for on_epoch_end. To contribute a test please refer to this guide and this discussion.

\n", "\n", "Called at the end of an epoch. " ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(KappaScore.on_epoch_end)" ] }, { "cell_type": "code", "execution_count": 39, "metadata": {}, "outputs": [ { "data": { "text/markdown": [ "

on_epoch_end[source][test]

\n", "\n", "> on_epoch_end(**`last_metrics`**, **\\*\\*`kwargs`**)\n", "\n", "
×

No tests found for on_epoch_end. To contribute a test please refer to this guide and this discussion.

\n", "\n", "Called at the end of an epoch. " ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(MatthewsCorreff.on_epoch_end)" ] }, { "cell_type": "code", "execution_count": 40, "metadata": {}, "outputs": [ { "data": { "text/markdown": [ "

on_train_begin[source][test]

\n", "\n", "> on_train_begin(**\\*\\*`kwargs`**)\n", "\n", "
×

No tests found for on_train_begin. To contribute a test please refer to this guide and this discussion.

\n", "\n", "To initialize constants in the callback. " ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(FBeta.on_train_begin)" ] }, { "cell_type": "code", "execution_count": 41, "metadata": {}, "outputs": [ { "data": { "text/markdown": [ "

on_epoch_end[source][test]

\n", "\n", "> on_epoch_end(**`last_metrics`**, **\\*\\*`kwargs`**)\n", "\n", "
×

No tests found for on_epoch_end. To contribute a test please refer to this guide and this discussion.

\n", "\n", "Called at the end of an epoch. " ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(RMSE.on_epoch_end)" ] }, { "cell_type": "code", "execution_count": 42, "metadata": {}, "outputs": [ { "data": { "text/markdown": [ "

on_train_begin[source][test]

\n", "\n", "> on_train_begin(**\\*\\*`kwargs`**)\n", "\n", "
×

No tests found for on_train_begin. To contribute a test please refer to this guide and this discussion.

\n", "\n", "To initialize constants in the callback. " ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(ConfusionMatrix.on_train_begin)" ] }, { "cell_type": "code", "execution_count": 43, "metadata": {}, "outputs": [ { "data": { "text/markdown": [ "

on_batch_end[source][test]

\n", "\n", "> on_batch_end(**`last_output`**:`Tensor`, **`last_target`**:`Tensor`, **\\*\\*`kwargs`**)\n", "\n", "
×

No tests found for on_batch_end. To contribute a test please refer to this guide and this discussion.

\n", "\n", "Called at the end of the batch. " ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(ConfusionMatrix.on_batch_end)" ] }, { "cell_type": "code", "execution_count": 44, "metadata": {}, "outputs": [ { "data": { "text/markdown": [ "

on_epoch_end[source][test]

\n", "\n", "> on_epoch_end(**\\*\\*`kwargs`**)\n", "\n", "
×

No tests found for on_epoch_end. To contribute a test please refer to this guide and this discussion.

\n", "\n", "Called at the end of an epoch. " ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(ConfusionMatrix.on_epoch_end)" ] }, { "cell_type": "code", "execution_count": 45, "metadata": {}, "outputs": [ { "data": { "text/markdown": [ "

on_epoch_end[source][test]

\n", "\n", "> on_epoch_end(**`last_metrics`**, **\\*\\*`kwargs`**)\n", "\n", "
×

No tests found for on_epoch_end. To contribute a test please refer to this guide and this discussion.

\n", "\n", "Called at the end of an epoch. " ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(Recall.on_epoch_end)" ] }, { "cell_type": "code", "execution_count": 46, "metadata": {}, "outputs": [ { "data": { "text/markdown": [ "

on_epoch_end[source][test]

\n", "\n", "> on_epoch_end(**`last_metrics`**, **\\*\\*`kwargs`**)\n", "\n", "
×

No tests found for on_epoch_end. To contribute a test please refer to this guide and this discussion.

\n", "\n", "Called at the end of an epoch. " ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(ExplainedVariance.on_epoch_end)" ] }, { "cell_type": "code", "execution_count": 47, "metadata": {}, "outputs": [ { "data": { "text/markdown": [ "

on_epoch_end[source][test]

\n", "\n", "> on_epoch_end(**`last_metrics`**, **\\*\\*`kwargs`**)\n", "\n", "
×

No tests found for on_epoch_end. To contribute a test please refer to this guide and this discussion.

\n", "\n", "Called at the end of an epoch. " ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(ExpRMSPE.on_epoch_end)" ] }, { "cell_type": "code", "execution_count": 48, "metadata": {}, "outputs": [ { "data": { "text/markdown": [ "

on_epoch_begin[source][test]

\n", "\n", "> on_epoch_begin(**\\*\\*`kwargs`**)\n", "\n", "
×

No tests found for on_epoch_begin. To contribute a test please refer to this guide and this discussion.

\n", "\n", "At the beginning of each epoch. " ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(ConfusionMatrix.on_epoch_begin)" ] }, { "cell_type": "code", "execution_count": 49, "metadata": {}, "outputs": [ { "data": { "text/markdown": [ "

on_epoch_end[source][test]

\n", "\n", "> on_epoch_end(**`last_metrics`**, **\\*\\*`kwargs`**)\n", "\n", "
×

No tests found for on_epoch_end. To contribute a test please refer to this guide and this discussion.

\n", "\n", "Called at the end of an epoch. " ], "text/plain": [ "" ] }, "metadata": {}, "output_type": "display_data" } ], "source": [ "show_doc(R2Score.on_epoch_end)" ] }, { "cell_type": "markdown", "metadata": {}, "source": [ "## New Methods - Please document or move to the undocumented section" ] } ], "metadata": { "jekyll": { "keywords": "fastai", "summary": "Useful metrics for training", "title": "metrics" }, "kernelspec": { "display_name": "Python 3", "language": "python", "name": "python3" }, "language_info": { "codemirror_mode": { "name": "ipython", "version": 3 }, "file_extension": ".py", "mimetype": "text/x-python", "name": "python", "nbconvert_exporter": "python", "pygments_lexer": "ipython3", "version": "3.8.2" } }, "nbformat": 4, "nbformat_minor": 2 }